Skip to content

Commit 7e72090

Browse files
committed
Add ops for rcnn models and UT
1 parent 41bc6d5 commit 7e72090

File tree

5 files changed

+356
-79
lines changed

5 files changed

+356
-79
lines changed

tests/run_pretrained_models.yaml

Lines changed: 14 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -374,3 +374,17 @@ style-transfer:
374374
outputs:
375375
- output:0
376376

377+
faster_rcnn_inception_v2_coco:
378+
url: http://download.tensorflow.org/models/object_detection/faster_rcnn_inception_v2_coco_2018_01_28.tar.gz
379+
model: faster_rcnn_inception_v2_coco_2018_01_28/frozen_inference_graph.pb
380+
opset_constraints:
381+
"onnx":
382+
"min": 11
383+
input_get: get_beach
384+
inputs:
385+
"image_tensor:0": [1, 224, 224, 3]
386+
outputs:
387+
- detection_boxes:0
388+
- detection_classes:0
389+
- detection_scores:0
390+
- num_detections:0

tests/test_backend.py

Lines changed: 71 additions & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -36,6 +36,8 @@
3636
_INPUT1 = "input1:0"
3737
_TFINPUT2 = "input2"
3838
_INPUT2 = "input2:0"
39+
_TFINPUT3 = "input3"
40+
_INPUT3 = "input3:0"
3941
_TFOUTPUT = "output"
4042
_OUTPUT = "output:0"
4143
_TFOUTPUT1 = "output1"
@@ -2470,6 +2472,74 @@ def test_batch_to_spacend(self):
24702472
_ = tf.batch_to_space_nd(input_x, block_size, crop, name=_TFOUTPUT)
24712473
self._run_test_case([_OUTPUT], {_INPUT: input_val})
24722474

2475+
@check_opset_min_version(11, "BatchToSpaceND")
2476+
def test_batch_to_spacend_non_const(self):
2477+
input_x_val = np.random.random_sample([40, 3, 5, 100]).astype(np.float32) # NHWC
2478+
block_shape_val = np.array([2, 2]).astype(np.int64)
2479+
crops_val = np.array([[1, 0], [2, 1]]).astype(np.int64)
2480+
input_x = tf.placeholder(dtype=tf.float32, shape=input_x_val.shape, name=_TFINPUT)
2481+
block_shape = tf.placeholder(dtype=tf.int64, shape=block_shape_val.shape, name=_TFINPUT1)
2482+
crops = tf.placeholder(dtype=tf.int64, shape=crops_val.shape, name=_TFINPUT2)
2483+
_ = tf.batch_to_space_nd(input_x, block_shape, crops, name=_TFOUTPUT)
2484+
self._run_test_case([_OUTPUT], {_INPUT: input_x_val, _INPUT1: block_shape_val, _INPUT2: crops_val})
2485+
2486+
@check_opset_min_version(11, "SpaceToBatchND")
2487+
def test_space_to_batchnd_non_const(self):
2488+
input_x_val = np.random.random_sample([40, 5, 7, 66]).astype(np.float32) # NHWC
2489+
block_size_val = np.array([2, 2]).astype(np.int64)
2490+
pad_val = np.array([[0, 1], [2, 1]]).astype(np.int64)
2491+
input_x = tf.placeholder(dtype=tf.float32, shape=input_x_val.shape, name=_TFINPUT)
2492+
block_size = tf.placeholder(dtype=tf.int64, shape=block_size_val.shape, name=_TFINPUT1)
2493+
pad = tf.placeholder(dtype=tf.int64, shape=pad_val.shape, name=_TFINPUT2)
2494+
_ = tf.space_to_batch_nd(input_x, block_size, pad, name=_TFOUTPUT)
2495+
self._run_test_case([_OUTPUT], {_INPUT: input_x_val, _INPUT1: block_size_val, _INPUT2: pad_val})
2496+
2497+
@check_opset_min_version(11, "CropAndResize")
2498+
def test_crop_and_resize_linear(self):
2499+
input_x_val = np.random.randint(low=0, high=256, size=[2, 36, 36, 3]).astype(np.float32) # NHWC
2500+
boxes_val = np.array([[0.5, 0.7, 0.7, 0.9], [0.2, 0.4, 0.4, 0.6]]).astype(np.float32)
2501+
box_ind_val = np.array([1, 0]).astype(np.int32)
2502+
corp_size_val = np.array([20, 20]).astype(np.int32)
2503+
input_x = tf.placeholder(dtype=tf.float32, shape=input_x_val.shape, name=_TFINPUT)
2504+
boxes = tf.placeholder(dtype=tf.float32, shape=boxes_val.shape, name=_TFINPUT1)
2505+
box_ind = tf.placeholder(dtype=tf.int32, shape=box_ind_val.shape, name=_TFINPUT2)
2506+
corp_size = tf.placeholder(dtype=tf.int32, shape=corp_size_val.shape, name=_TFINPUT3)
2507+
_ = tf.image.crop_and_resize(input_x, boxes, box_ind, corp_size, name=_TFOUTPUT, method='bilinear')
2508+
self._run_test_case([_OUTPUT],
2509+
{_INPUT: input_x_val, _INPUT1: boxes_val, _INPUT2: box_ind_val, _INPUT3: corp_size_val},
2510+
rtol=1e-05, atol=1e-04)
2511+
2512+
@check_tf_min_version("1.9")
2513+
@check_opset_min_version(11, "CropAndResize")
2514+
def test_crop_and_resize_nearest(self):
2515+
input_x_val = np.random.randint(low=0, high=256, size=[1, 36, 36, 3]).astype(np.float32) # NHWC
2516+
boxes_val = np.array([[0.2, 0.4, 0.6, 0.8]]).astype(np.float32)
2517+
box_ind_val = np.array([0]).astype(np.int32)
2518+
corp_size_val = np.array([30, 30]).astype(np.int32)
2519+
input_x = tf.placeholder(dtype=tf.float32, shape=input_x_val.shape, name=_TFINPUT)
2520+
boxes = tf.placeholder(dtype=tf.float32, shape=boxes_val.shape, name=_TFINPUT1)
2521+
box_ind = tf.placeholder(dtype=tf.int32, shape=box_ind_val.shape, name=_TFINPUT2)
2522+
corp_size = tf.placeholder(dtype=tf.int32, shape=corp_size_val.shape, name=_TFINPUT3)
2523+
_ = tf.image.crop_and_resize(input_x, boxes, box_ind, corp_size, name=_TFOUTPUT, method='nearest')
2524+
self._run_test_case([_OUTPUT],
2525+
{_INPUT: input_x_val, _INPUT1: boxes_val, _INPUT2: box_ind_val, _INPUT3: corp_size_val},
2526+
rtol=1e-05, atol=1e-04)
2527+
2528+
@check_opset_min_version(11, "CropAndResize")
2529+
def test_crop_and_resize_extrapolation(self):
2530+
input_x_val = np.random.randint(low=0, high=256, size=[1, 36, 36, 3]).astype(np.float32) # NHWC
2531+
boxes_val = np.array([[0.2, 0.4, 1.2, 1.4]]).astype(np.float32)
2532+
box_ind_val = np.array([0]).astype(np.int32)
2533+
corp_size_val = np.array([40, 40]).astype(np.int32)
2534+
input_x = tf.placeholder(dtype=tf.float32, shape=input_x_val.shape, name=_TFINPUT)
2535+
boxes = tf.placeholder(dtype=tf.float32, shape=boxes_val.shape, name=_TFINPUT1)
2536+
box_ind = tf.placeholder(dtype=tf.int32, shape=box_ind_val.shape, name=_TFINPUT2)
2537+
corp_size = tf.placeholder(dtype=tf.int32, shape=corp_size_val.shape, name=_TFINPUT3)
2538+
_ = tf.image.crop_and_resize(input_x, boxes, box_ind, corp_size, name=_TFOUTPUT, extrapolation_value=1.0)
2539+
self._run_test_case([_OUTPUT],
2540+
{_INPUT: input_x_val, _INPUT1: boxes_val, _INPUT2: box_ind_val, _INPUT3: corp_size_val},
2541+
rtol=1e-04, atol=1e-03)
2542+
24732543
def test_batch_to_space3d(self):
24742544
block_size = [2, 2]
24752545
crop = [[0, 1], [2, 1]]
@@ -2828,9 +2898,8 @@ def test_unique(self):
28282898
_ = tf.identity(x1_, name=_TFOUTPUT)
28292899
_ = tf.identity(x2_, name=_TFOUTPUT1)
28302900
# FIXME: indices in onnx are not the same as in tensorflow so don't check for now
2831-
#self._run_test_case([_OUTPUT, _OUTPUT1], {_INPUT: x_val})
2901+
# self._run_test_case([_OUTPUT, _OUTPUT1], {_INPUT: x_val})
28322902
self._run_test_case([_OUTPUT], {_INPUT: x_val})
28332903

2834-
28352904
if __name__ == '__main__':
28362905
unittest_main()

tf2onnx/onnx_opset/nn.py

Lines changed: 65 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -534,6 +534,71 @@ def version_11(cls, ctx, node, **kwargs):
534534
cls.version_1(ctx, node, **kwargs)
535535

536536

537+
@tf_op(["CropAndResize"])
538+
class CropAndResize:
539+
@classmethod
540+
def version_11(cls, ctx, node, **kwargs):
541+
# create loop of resize to cater to tensorflow CropAndResize, one box one iteration
542+
mode = "nearest" if node.get_attr("method") is not None and node.get_attr(
543+
"method").s == b"nearest" else "linear"
544+
extrapolation_value = float(node.get_attr("extrapolation_value", "0").f)
545+
input_x = node.inputs[0]
546+
boxes = node.inputs[1]
547+
box_ind = node.inputs[2]
548+
crop_size = node.inputs[3]
549+
trip_name = utils.make_name(node.name + "_i")
550+
cond_name = utils.make_name(node.name + "_cond")
551+
cond_out_name = utils.make_name(node.name + "cond_out")
552+
g = ctx.create_new_graph_with_same_config()
553+
g.add_graph_input(trip_name, TensorProto.INT64, [1])
554+
g.add_graph_input(cond_name, TensorProto.BOOL, [])
555+
g.parent_graph = ctx
556+
const_zero = g.make_const(utils.make_name(node.name + "_const_zero"), np.array([0], dtype=np.int32))
557+
const_zero_long = g.make_const(utils.make_name(node.name + "_const_zero_long"), np.array([0], dtype=np.int64))
558+
const_one = g.make_const(utils.make_name(node.name + "_const_one"), np.array([1], dtype=np.int32))
559+
const_one_long = g.make_const(utils.make_name(node.name + "_const_one_long"), np.array([1], dtype=np.int64))
560+
index_end = g.make_node("Add", [trip_name, const_one_long.output[0]])
561+
box_index_from = g.make_node("Slice", [box_ind.output[0], trip_name, index_end.output[0]], name="Slice_a")
562+
box_index_to = g.make_node("Add", [box_index_from.output[0], const_one.output[0]])
563+
target_x = g.make_node("Slice", [input_x.output[0], box_index_from.output[0], box_index_to.output[0],
564+
const_zero.output[0]], name="Slice_b")
565+
transposed_x = g.make_node("Transpose", [target_x.output[0]], attr={'perm': constants.NHWC_TO_NCHW})
566+
shape_of_transposed_x = g.make_node("Shape", [transposed_x.output[0]])
567+
const_zero_zero = g.make_const(utils.make_name(node.name + "_const_zero_zero"),
568+
np.array([0, 0], dtype=np.float32))
569+
const_one_one = g.make_const(utils.make_name(node.name + "_const_one_one"),
570+
np.array([1, 1], dtype=np.float32))
571+
const_four = g.make_const(utils.make_name(node.name + "_const_four"), np.array([4], dtype=np.int64))
572+
const_empty_float = g.make_const(utils.make_name("const_empty_float"), np.array([], dtype=np.float32))
573+
first_half_of_shape = GraphBuilder(g).make_slice(
574+
{"data": shape_of_transposed_x.output[0], "ends": [2], "starts": [0]})
575+
box = g.make_node("Slice", [boxes.output[0], trip_name, index_end.output[0], const_zero_long.output[0]],
576+
name="Slice_c")
577+
roi_raw = g.make_node("Reshape", [box.output[0], const_four.output[0]])
578+
roi_raw_first_half = GraphBuilder(g).make_slice({"data": roi_raw.output[0], "ends": [2], "starts": [0]})
579+
roi_raw_second_half = GraphBuilder(g).make_slice({"data": roi_raw.output[0], "ends": [4], "starts": [2]})
580+
roi_concat_1 = g.make_node("Concat", [const_zero_zero.output[0], roi_raw_first_half], attr={'axis': 0})
581+
roi_concat_2 = g.make_node("Concat", [const_one_one.output[0], roi_raw_second_half], attr={'axis': 0})
582+
final_roi = g.make_node("Concat", [roi_concat_1.output[0], roi_concat_2.output[0]], attr={'axis': 0})
583+
crop_size_int64 = g.make_node("Cast", [crop_size.output[0]], attr={'to': TensorProto.INT64})
584+
final_crop_size = g.make_node("Concat", [first_half_of_shape, crop_size_int64.output[0]], {'axis': 0})
585+
resized_x = g.make_node("Resize", [transposed_x.output[0], final_roi.output[0], const_empty_float.output[0],
586+
final_crop_size.output[0]],
587+
attr={"mode": mode, "extrapolation_value": extrapolation_value,
588+
"coordinate_transformation_mode": "tf_crop_and_resize"})
589+
recovered_x = g.make_node("Transpose", [resized_x.output[0]], attr={'perm': constants.NCHW_TO_NHWC})
590+
squeeze_x = g.make_node("Squeeze", inputs=[recovered_x.output[0]], attr={"axes": [0]})
591+
g.make_node("Identity", [cond_name], outputs=[cond_out_name])
592+
g.add_graph_output(cond_out_name, TensorProto.BOOL, [])
593+
g.add_graph_output(squeeze_x.output[0], ctx.get_dtype(node.input[0]), [-1, -1, -1])
594+
trip_node = ctx.make_node("Size", [box_ind.output[0]])
595+
cond_const = ctx.make_const(utils.make_name("cond"), np.ones((), dtype=np.bool))
596+
ctx.remove_node(node.name)
597+
inner_loop = ctx.make_node("Loop", [trip_node.output[0], cond_const.output[0]], name=node.name,
598+
outputs=node.output)
599+
inner_loop.set_body_graph_as_attr("body", g)
600+
601+
537602
@tf_op(["ResizeBilinear", "ResizeNearestNeighbor"])
538603
class Resize:
539604
@classmethod

0 commit comments

Comments
 (0)