1 2 3 4 5 6 7 8 9 10 11 12 13 14 15 16 17 18 19 20 21 22 23 24 25 26 27 28 29 30 31 32 33 34 35 36 37 38 39 40 41 42 43 44 45 46 47 48 49 50 51 52 53 54 55 56 57 58 59 60 61 62 63 64 65 66 67 68 69 70 71 72 73 74 75 76 77 78 79 80 81 82 83 84 85 86 87 88 89 90 91 92 93 94 95 96 97 98 99 100 101 102 103 104 105 106 107 108 109 110 111 112 113 114 115 116 117 118 119 120 121 122 123 124 125 126 127 128 129 130 131 132 133 134 135 136 137 138 139 140 141 142 143 144 145 146 147 148 149 150 151 152 153 154 155 156 157 158 159 160 161 162 163 164 165 166 167 168 169 170 171 172 173 174 175 176 177 178 179 180 181 182 183 184 185 186 187 188 189 190 191 192 193 194 195 196 197 198 199 200 201 202 203 204 205 206 207 208 209 210 211 212 213 214 215 216 217 218 219 220 221 222 223 224 225 226 227 228 229 230 231 232 233 234 235 236 237 238 239 240 241 242 243 244 245 246 247 248 249 250 251 252 253 254 255 256 257 258 259 260 261 262 263 264 265 266 267 268 269 270 271 272 273 274 275 276 277 278 279 280 281 282 283 284 285 286 287 288 289 290 291 292 293 294 295 296 297 298 299 300 301 302 303 304 305 306 307 308 309 310 311 312 313 314 315 316 317 318 319 320 321 322
|
import caffe2.python.fakelowp.init_shared_libs # noqa
import numpy as np
from caffe2.python import core, workspace
from caffe2.python.onnx.onnxifi import onnxifi_caffe2_net
from hypothesis import given, strategies as st, settings
from caffe2.python.fakelowp.test_utils import print_test_debug_info
import caffe2.python.serialized_test.serialized_test_util as serial
import datetime
core.GlobalInit(["caffe2",
"--caffe2_log_level=-3",
"--glow_global_fp16=1",
"--glow_clip_quant_range_to_fp16=1",
"--glow_global_fp16_constants=1"
])
class Int8OpsTest(serial.SerializedTestCase):
def _get_scale_zp(self, tensor):
tensor_max = np.max(tensor)
tensor_min = min(0, np.min(tensor))
scale = np.float32(np.float16((tensor_max - tensor_min) / 255.0))
if scale < 1e-6:
scale = np.float32(1e-6)
zero_point = 0 - tensor_min / scale
zero_point = int(round(np.clip(zero_point, 0, 255.0)))
return (scale, zero_point)
@given(
n=st.integers(2, 1024),
rand_seed=st.integers(0, 65534),
non_zero_offset=st.booleans()
)
@settings(deadline=datetime.timedelta(seconds=50))
def test_int8_quantize(self, n, rand_seed, non_zero_offset):
print("n={}, rand_seed={}".format(n, rand_seed))
np.random.seed(rand_seed)
workspace.ResetWorkspace()
if non_zero_offset:
X_fp32 = np.random.uniform(-1, 1, size=(n, n)).astype(np.float16) \
.astype(np.float32)
else:
X_fp32 = np.random.rand(n, n).astype(np.float16).astype(np.float32)
W_fp32 = np.identity(n, dtype=np.float32)
b_fp32 = np.zeros((n,), dtype=np.float32)
X_scale, X_zero_point = self._get_scale_zp(X_fp32)
workspace.FeedBlob("X", X_fp32)
workspace.FeedBlob("W", W_fp32)
workspace.FeedBlob("b", b_fp32)
workspace.RunOperatorOnce(
core.CreateOperator(
"Int8FCPackWeight",
["W"],
["W_int8"],
engine="DNNLOWP",
save_unpacked_weights=True,
in_scale=X_scale,
)
)
ref_net = core.Net("net")
ref_net.Int8QuantizeNNPI(
["X"],
["X_int8"],
Y_scale=X_scale,
Y_zero_point=X_zero_point
)
ref_net.Int8FCFakeAcc32NNPI(
["X_int8", "W_int8", "b"],
["Y_int8"],
Y_scale=X_scale,
Y_zero_point=X_zero_point,
)
ref_net.Int8DequantizeNNPI(
["Y_int8"],
["Y"]
)
ref_net.Proto().external_output.append("Y")
# run ref_net
workspace.RunNetOnce(ref_net)
Y_fbgemm = workspace.FetchBlob("Y")
# run onnxifi net
ref_net.Proto().op[0].type = "Int8Quantize"
ref_net.Proto().op[1].type = "Int8FC"
ref_net.Proto().op[2].type = "Int8Dequantize"
net_onnxified = onnxifi_caffe2_net(
ref_net.Proto(),
{},
debug=True,
adjust_batch=False,
use_onnx=False,
weight_names=["W_int8", "b"],
)
num_onnxified_ops = sum(
1 if o.type == "Onnxifi" else 0 for o in net_onnxified.op
)
np.testing.assert_equal(num_onnxified_ops, 1)
workspace.CreateNet(net_onnxified)
workspace.RunNet(net_onnxified.name)
Y_glow = workspace.FetchBlob("Y")
if not np.allclose(Y_glow, Y_fbgemm):
diff_Y = np.abs(Y_glow - Y_fbgemm)
print_test_debug_info(
"int8_fc",
{
"seed": rand_seed,
"n": n,
"X": X_fp32,
"W": W_fp32,
"b": b_fp32,
"Y_fbgemm": Y_fbgemm,
"Y_glow": Y_glow,
"diff": diff_Y,
"maxdiff": diff_Y.max(axis=1),
},
)
assert 0
@given(
n=st.integers(1, 1024),
m=st.integers(1, 1024),
k=st.integers(1, 1024),
f=st.integers(1, 1), # TODO: figure a safe number to increase
rand_seed=st.integers(0, 65534),
quantize_bias=st.sampled_from([False]),
)
@settings(deadline=datetime.timedelta(seconds=50))
def test_int8_fc(
self, n, m, k, rand_seed, quantize_bias, f
):
print(
f"n={n}, m={m}, k={k}, rand_seed={rand_seed}, quantize_bias={quantize_bias}"
)
np.random.seed(rand_seed)
workspace.ResetWorkspace()
ff = float(f)
X_fp32 = np.random.uniform(-ff, ff, size=(m, k)).astype(np.float32)
W_fp32 = np.random.uniform(-ff, ff, size=(n, k)).astype(np.float32)
b_fp32 = np.random.uniform(-ff, ff, size=(n)).astype(np.float32)
X_scale, X_zero_point = self._get_scale_zp(X_fp32)
Y_fp32 = np.dot(X_fp32, W_fp32.T) + b_fp32
Y_scale, Y_zero_point = self._get_scale_zp(Y_fp32)
workspace.FeedBlob("X", X_fp32)
workspace.FeedBlob("W", W_fp32)
workspace.FeedBlob("b", b_fp32)
workspace.RunOperatorOnce(
core.CreateOperator(
"Int8FCPackWeight",
["W", "b"] if quantize_bias else ["W"],
["W_int8", "b_int32"] if quantize_bias else ["W_int8"],
engine="DNNLOWP",
save_unpacked_weights=True,
in_scale=X_scale,
)
)
ref_net = core.Net("net")
ref_net.Int8QuantizeNNPI(
["X"],
["X_int8"],
Y_scale=X_scale,
Y_zero_point=X_zero_point
)
ref_net.Int8FCFakeAcc32NNPI(
["X_int8", "W_int8", "b_int32" if quantize_bias else "b"],
["Y_int8"],
Y_scale=Y_scale,
Y_zero_point=Y_zero_point,
)
ref_net.Int8DequantizeNNPI(
["Y_int8"],
["Y"]
)
ref_net.Proto().external_output.append("Y")
# run ref_net
workspace.RunNetOnce(ref_net)
Y_fbgemm = workspace.FetchBlob("Y")
# run onnxifi net
ref_net.Proto().op[0].type = "Int8Quantize"
ref_net.Proto().op[1].type = "Int8FC"
ref_net.Proto().op[2].type = "Int8Dequantize"
net_onnxified = onnxifi_caffe2_net(
ref_net.Proto(),
{},
debug=True,
adjust_batch=False,
use_onnx=False,
weight_names=["W_int8", "b_int32"] if quantize_bias else ["W_int8", "b"],
)
num_onnxified_ops = sum(
1 if o.type == "Onnxifi" else 0 for o in net_onnxified.op
)
np.testing.assert_equal(num_onnxified_ops, 1)
workspace.CreateNet(net_onnxified)
workspace.RunNet(net_onnxified.name)
Y_glow = workspace.FetchBlob("Y")
if not np.allclose(Y_glow, Y_fbgemm):
diff_Y = np.abs(Y_glow - Y_fbgemm)
print_test_debug_info(
"int8_fc",
{
"seed": rand_seed,
"n": n,
"m": m,
"k": k,
"X": X_fp32,
"W": W_fp32,
"b": b_fp32,
"Y_fbgemm": Y_fbgemm,
"Y_glow": Y_glow,
"diff": diff_Y,
"maxdiff": diff_Y.max(axis=1),
},
)
assert 0
@given(
n=st.integers(1, 4),
rand_seed=st.integers(0, 65534)
)
@settings(deadline=datetime.timedelta(seconds=10))
def test_int8_small_input(self, n, rand_seed):
print("n={}, rand_seed={}".format(n, rand_seed))
np.random.seed(rand_seed)
workspace.ResetWorkspace()
X_fp32 = np.random.uniform(0.01, 0.03, size=(n, n)).astype(np.float32)
W_fp32 = np.identity(n, dtype=np.float32)
b_fp32 = np.zeros((n,), dtype=np.float32)
X_scale, X_zero_point = self._get_scale_zp(X_fp32)
workspace.FeedBlob("X", X_fp32)
workspace.FeedBlob("W", W_fp32)
workspace.FeedBlob("b", b_fp32)
workspace.RunOperatorOnce(
core.CreateOperator(
"Int8FCPackWeight",
["W"],
["W_int8"],
engine="DNNLOWP",
save_unpacked_weights=True,
in_scale=X_scale,
)
)
ref_net = core.Net("net")
ref_net.Int8QuantizeNNPI(
["X"],
["X_int8"],
Y_scale=X_scale,
Y_zero_point=X_zero_point
)
ref_net.Int8FCFakeAcc32NNPI(
["X_int8", "W_int8", "b"],
["Y_int8"],
Y_scale=X_scale,
Y_zero_point=X_zero_point,
)
ref_net.Int8DequantizeNNPI(
["Y_int8"],
["Y"]
)
ref_net.Proto().external_output.append("Y")
# run ref_net
workspace.RunNetOnce(ref_net)
Y_fbgemm = workspace.FetchBlob("Y")
# run onnxifi net
ref_net.Proto().op[0].type = "Int8Quantize"
ref_net.Proto().op[1].type = "Int8FC"
ref_net.Proto().op[2].type = "Int8Dequantize"
net_onnxified = onnxifi_caffe2_net(
ref_net.Proto(),
{},
debug=True,
adjust_batch=False,
use_onnx=False,
weight_names=["W_int8", "b"],
)
num_onnxified_ops = sum(
1 if o.type == "Onnxifi" else 0 for o in net_onnxified.op
)
np.testing.assert_equal(num_onnxified_ops, 1)
workspace.CreateNet(net_onnxified)
workspace.RunNet(net_onnxified.name)
Y_glow = workspace.FetchBlob("Y")
if not np.allclose(Y_glow, Y_fbgemm):
diff_Y = np.abs(Y_glow - Y_fbgemm)
print_test_debug_info(
"int8_fc",
{
"seed": rand_seed,
"n": n,
"X": X_fp32,
"W": W_fp32,
"b": b_fp32,
"Y_fbgemm": Y_fbgemm,
"Y_glow": Y_glow,
"diff": diff_Y,
"maxdiff": diff_Y.max(axis=1),
},
)
assert 0
|