Caffe2 - Python API
A deep learning, cross platform ML framework
backend.py
1 ## @package onnx
2 # Module caffe2.python.onnx.backend
3 
4 """Backend for running ONNX on Caffe2
5 
6 To run this, you will need to have Caffe2 installed as well.
7 """
8 from __future__ import absolute_import
9 from __future__ import division
10 from __future__ import print_function
11 from __future__ import unicode_literals
12 
13 import os
14 import collections
15 from subprocess import Popen, PIPE
16 import zipfile
17 import itertools
18 
19 # When onnx is built against a version of protobuf that is older than
20 # that which is vendored with caffe2, onnx will crash if caffe2's
21 # vendored protobuf is loaded first. We can work around this by
22 # importing onnx first, which will cause it to go out and pick up the
23 # system protobuf.
24 import onnx.backend
25 
26 import caffe2
27 from caffe2.python import core, workspace, rnn_cell, gru_cell
28 from caffe2.python.compatibility import container_abcs
29 from caffe2.python.model_helper import ModelHelper
30 from caffe2.proto import caffe2_pb2
32 import numpy as np
33 import onnx
34 from onnx import checker, GraphProto, TensorProto, AttributeProto, ModelProto
35 import onnx.numpy_helper
36 import onnx.defs
37 import onnx.optimizer
38 import onnx.shape_inference
39 import onnx.utils
40 from onnx.backend.base import Backend, Device, DeviceType, namedtupledict
41 
42 from caffe2.python.onnx.workspace import Workspace
43 from caffe2.python.onnx.backend_rep import Caffe2Rep
44 from caffe2.python.onnx.backend_cpp_rep import Caffe2CppRep
45 
47 
48 import warnings
49 
50 def force_unicode(s):
51  try:
52  return s.decode('utf-8')
53  except AttributeError:
54  return s
55 
56 def get_device_option(device):
57  m = {DeviceType.CPU: caffe2_pb2.CPU,
58  DeviceType.CUDA: workspace.GpuDeviceType}
59  return core.DeviceOption(m[device.type], device.device_id)
60 
61 
62 class OnnxAttributes(dict):
63  """
64  This is a more convenient way to work with ONNX/Caffe2 attributes
65  that is not the protobuf representation.
66  """
67  @staticmethod
68  def from_onnx(args):
69  d = OnnxAttributes()
70  for arg in args:
71  d[arg.name] = convertAttributeProto(arg)
72  return d
73 
74  def caffe2(self, kmap=lambda k: k):
75  for k, v in self.items():
76  if kmap(k) != '':
77  yield caffe2.python.utils.MakeArgument(kmap(k), v)
78 
79 # TODO: Move this into ONNX main library
80 def convertAttributeProto(onnx_arg):
81  """
82  Convert an ONNX AttributeProto into an appropriate Python object
83  for the type.
84 
85  NB: Tensor attribute gets returned as the straight proto.
86  """
87  if onnx_arg.HasField('f'):
88  return onnx_arg.f
89  elif onnx_arg.HasField('i'):
90  return onnx_arg.i
91  elif onnx_arg.HasField('s'):
92  return onnx_arg.s
93  elif onnx_arg.HasField('t'):
94  return onnx_arg.t # this is a proto!
95  elif onnx_arg.HasField('g'):
96  return Caffe2Backend._graph_to_net(onnx_arg.g, Caffe2Backend._known_opset_version)
97  elif len(onnx_arg.floats):
98  return list(onnx_arg.floats)
99  elif len(onnx_arg.ints):
100  return list(onnx_arg.ints)
101  elif len(onnx_arg.strings):
102  return list(onnx_arg.strings)
103  elif len(onnx_arg.graphs):
104  retval = []
105  # TODO: this doesn't work with RNN ops
106  for g in onnx_arg.graphs:
107  retval.append(Caffe2Backend._graph_to_net(g, Caffe2Backend._known_opset_version))
108  return retval
109  else:
110  raise ValueError("Unsupported ONNX attribute: {}".format(onnx_arg))
111 
112 
113 # TODO: Move this into ONNX main library
114 class OnnxNode(object):
115  """
116  Reimplementation of NodeProto from ONNX, but in a form
117  more convenient to work with from Python.
118 
119  We may temporarily edit these nodes to get them into Caffe2 form,
120  before actually translating into the Caffe2 protobuf, since this
121  is easier than decomposing everything, and putting it back together
122  when we're ready.
123  """
124  def __init__(self, node):
125  self.name = str(node.name)
126  self.op_type = str(node.op_type)
127  self.attrs = OnnxAttributes.from_onnx(node.attribute)
128  self.inputs = list(node.input)
129  self.outputs = list(node.output)
130 
131 
132 Caffe2Ops = collections.namedtuple('Caffe2Ops', ['ops', 'init_ops', 'interface_blobs'])
133 
134 
135 class Caffe2Backend(Backend):
136 
137  # The greatest version of the ONNX operator set which we are aware of.
138  # Models whose version is larger than this will cause us to emit a warning
139  # that we are attempting to translate on a "best effort" basis.
140  #
141  # If you increase this, make SURE you cross-reference all BC-breaking
142  # changes from one version to the next, and any that you did not
143  # implement, mark as broken in _broken_operators
144  _known_opset_version = 9
145 
146  # This dictionary will record operators which are KNOWN to be
147  # broken, so we give a good error message rather than do something
148  # bogus and then fail.
149  _broken_operators = {
150  # 'BrokenOp': version_it_was_broken_in
151  }
152 
153  # Operators that are different between Caffe2 and
154  # ONNX but only in their name.
155  # In most cases, this should be empty - as the effort of ONNX is
156  # to unify the operator definitions.
157  _renamed_operators = {
158  'GlobalMaxPool': 'MaxPool',
159  'GlobalAveragePool': 'AveragePool',
160  'Pad': 'PadImage',
161  'Neg': 'Negative',
162  'BatchNormalization': 'SpatialBN',
163  'InstanceNormalization': 'InstanceNorm',
164  'MatMul': 'BatchMatMul',
165  'Upsample': 'ResizeNearest',
166  'Identity': 'Copy',
167  'InstanceNormalization': 'InstanceNorm',
168  'Equal': 'EQ',
169  'Less': 'LT',
170  'Greater': 'GT',
171  'Unsqueeze': 'ExpandDims',
172  'Loop': 'ONNXWhile',
173  'Tile': 'NumpyTile',
174  'RandomNormal': 'GaussianFill',
175  }
176 
177  _global_renamed_attrs = {'kernel_shape': 'kernels'}
178  _per_op_renamed_attrs = {
179  'Squeeze': {'axes': 'dims'},
180  'Unsqueeze': {'axes': 'dims'},
181  'Transpose': {'perm': 'axes'},
182  'Upsample': {'mode': '',
183  'scales': ''},
184  'ConvTranspose': {'output_padding': 'adjs'},
185  'Selu': {'gamma': 'scale'},
186  'If': {'then_branch': 'then_net',
187  'else_branch': 'else_net'}
188  }
189 
190  # operators whose behavior is different beyond renaming
191  # the value is an attribute of this class that is a
192  # function from ToffeIR node_def to caffe2 op_def
193  _special_operators = {
194  'LSTM': '_create_rnn_variant',
195  'GRU': '_create_rnn_variant',
196  'RNN': '_create_rnn_variant',
197  'Loop': '_create_loop',
198  'If': '_create_if',
199  'Upsample': '_create_upsample',
200  'RandomNormal': '_create_gaussian_fill'
201  }
202 
203  # Dummy name generator
204  _dummy_name = C.DummyName()
205 
206  @classmethod
207  def dummy_name(cls):
208  return cls._dummy_name.new_dummy_name()
209 
210  # NB: By default, you will use the LATEST definition of the operator,
211  # so this interface MAY make BC-breaking changes. Specify an
212  # opset_version if you don't want this to version.
213  @classmethod
214  def run_node(cls, node, inputs, device='CPU', opset_version=_known_opset_version, outputs_info=None):
215  super(Caffe2Backend, cls).run_node(node, inputs, device=device,
216  outputs_info=outputs_info, opset_version=opset_version)
217 
218  value_infos = []
219  device_option = get_device_option(Device(device))
220  ws = Workspace()
221  with core.DeviceScope(device_option): # temporary!
222  if isinstance(inputs, dict):
223  for key, value in inputs.items():
224  ws.FeedBlob(key, value)
225  value_infos.append(onnx.helper.make_tensor_value_info(
226  name=key,
227  elem_type=onnx.mapping.NP_TYPE_TO_TENSOR_TYPE[value.dtype],
228  shape=value.shape).SerializeToString())
229  else:
230  assert len(node.input) == len(inputs), "{}: expected {} but got {}".format(
231  node.op_type, len(node.input), len(inputs))
232  for key, value in zip(node.input, inputs):
233  ws.FeedBlob(key, value)
234  value_infos.append(onnx.helper.make_tensor_value_info(
235  name=key,
236  elem_type=onnx.mapping.NP_TYPE_TO_TENSOR_TYPE[value.dtype],
237  shape=value.shape).SerializeToString())
238 
239  ops = []
240  cbackend = C.Caffe2Backend(cls._dummy_name)
241  ops_str = cbackend.convert_node(node.SerializeToString(), value_infos, opset_version)
242  for s in ops_str[0] + ops_str[1]:
243  op = caffe2_pb2.OperatorDef()
244  op.ParseFromString(s)
245  op.device_option.CopyFrom(device_option)
246  ops.append(op)
247  ws.RunOperatorsOnce(ops)
248  output_values = [ws.FetchBlob(name) for name in node.output]
249  return namedtupledict('Outputs', node.output)(*output_values)
250 
251  @classmethod
252  def _create_tensor_filling_op(cls, onnx_tensor, name=None):
253  """
254  Given an Onnx TensorProto, translate it into a Caffe2 operator
255  which produces the given tensor filling op.
256  """
257  assert name or onnx_tensor.name
258  name = name or onnx_tensor.name
259 
260  c2_op = caffe2_pb2.OperatorDef()
261 
262  c2_values = c2_op.arg.add()
263  c2_values.name = "values"
264 
265  def tensor2list(onnx_tensor):
266  # Use the onnx.numpy_helper because the data may be raw
267  return onnx.numpy_helper.to_array(onnx_tensor).flatten().tolist()
268 
269  if onnx_tensor.data_type in [TensorProto.FLOAT]:
270  c2_op.type = 'GivenTensorFill'
271  c2_values.floats.extend(tensor2list(onnx_tensor))
272  elif onnx_tensor.data_type in [TensorProto.DOUBLE]:
273  c2_op.type = 'GivenTensorDoubleFill'
274  c2_values.floats.extend(tensor2list(onnx_tensor))
275  elif onnx_tensor.data_type in [TensorProto.INT64,
276  TensorProto.UINT32]:
277  c2_op.type = 'GivenTensorInt64Fill'
278  c2_values.ints.extend(tensor2list(onnx_tensor))
279  elif onnx_tensor.data_type in [TensorProto.UINT8,
280  TensorProto.INT8,
281  TensorProto.UINT16,
282  TensorProto.INT16,
283  TensorProto.INT32]:
284  c2_op.type = 'GivenTensorIntFill'
285  c2_values.ints.extend(tensor2list(onnx_tensor))
286  elif onnx_tensor.data_type == TensorProto.BOOL:
287  c2_op.type = 'GivenTensorBoolFill'
288  c2_values.ints.extend(tensor2list(onnx_tensor))
289  elif onnx_tensor.data_type == TensorProto.STRING:
290  c2_op.type = 'GivenTensorStringFill'
291  c2_values.strings.extend(onnx_tensor.string_data)
292  else:
293  raise RuntimeError(
294  "unrecognized tensor type {}".format(onnx_tensor.data_type))
295 
296  c2_shape = c2_op.arg.add()
297  c2_shape.name = "shape"
298  c2_shape.ints.extend(onnx_tensor.dims)
299 
300  c2_op.output.append(name)
301 
302  return c2_op
303 
304  @classmethod
305  def _rnn_reform_weights(cls, reforms, name, hidden_size, init_net, gates, reorder_indices):
306  for name_from, name_to, do_concat, extra_dims in reforms:
307  gate_blobs = ['%s/%s_%s' % (name, prefix, name_to) for prefix in gates]
308  for i, x in enumerate(gate_blobs):
309  dim0 = i * hidden_size, (i+1) * hidden_size
310  starts, ends = zip(dim0, *extra_dims)
311  init_net.Slice(name_from, x, starts=starts, ends=ends)
312  if do_concat:
313  reordered_gate_blobs = [gate_blobs[i] for i in reorder_indices]
314  init_net.Concat(reordered_gate_blobs, ['%s/%s' % (name, name_to), cls.dummy_name()], axis=0)
315 
316  @classmethod
317  def _make_rnn_direction(cls, input_blob, B, W, R, initial_states_and_names, sequence_lens,
318  pred_mh, init_net,
319  input_size, hidden_size, num_gates, direction_offset,
320  Bi, Br, W_, R_,
321  reform, make_cell, keep_outputs):
322  name = cls.dummy_name()
323 
324  # input and recurrence biases are squashed together in onnx
325  # but not in caffe2
326  gates_hidden_size = num_gates * hidden_size
327  bias_offset = 2 * direction_offset * gates_hidden_size
328  weight_offset = direction_offset * gates_hidden_size
329  Bi = init_net.Slice(B, name + Bi,
330  starts=[bias_offset + 0 * gates_hidden_size],
331  ends =[bias_offset + 1 * gates_hidden_size])
332  Br = init_net.Slice(B, name + Br,
333  starts=[bias_offset + 1 * gates_hidden_size],
334  ends =[bias_offset + 2 * gates_hidden_size])
335  W_ = init_net.Slice(W, name + W_,
336  starts=[weight_offset + 0 * gates_hidden_size, 0],
337  ends =[weight_offset + 1 * gates_hidden_size,-1])
338  R_ = init_net.Slice(R, name + R_,
339  starts=[weight_offset + 0 * gates_hidden_size, 0],
340  ends =[weight_offset + 1 * gates_hidden_size,-1])
341 
342  initial_states_sliced = []
343  for initial_state, name_suffix in initial_states_and_names:
344  initial_states_sliced.append(
345  pred_mh.net.Slice(initial_state, name + name_suffix,
346  starts=[direction_offset + 0, 0, 0],
347  ends =[direction_offset + 1,-1,-1]))
348 
349  if direction_offset == 1:
350  if sequence_lens is not None:
351  seq_lens_for_reverse = sequence_lens
352  else:
353  input_shape = pred_mh.net.Shape(input_blob, name + '/input_shape')
354  batch_size = pred_mh.net.Slice(input_shape, name + '/batch_size_slice', starts=[1], ends=[2])
355  seq_len = pred_mh.net.Slice(input_shape, name + '/seq_len_slice', starts=[0], ends=[1])
356  dummy_sequence_lens = pred_mh.net.Tile([seq_len, batch_size], name + '/dummy_sequence_lens', axis=0)
357  pred_mh.net.Reshape(dummy_sequence_lens, [dummy_sequence_lens, cls.dummy_name()], shape=[-1])
358  seq_lens_for_reverse = pred_mh.net.Cast(dummy_sequence_lens, name + '/seq_lens_for_reverse', to=core.DataType.INT32)
359  reform(Bi, Br, W_, R_, name, hidden_size, init_net)
360 
361  if direction_offset == 1:
362  input = pred_mh.net.ReversePackedSegs(
363  [input_blob, seq_lens_for_reverse], name + "/input-reversed")
364  else:
365  input = input_blob
366 
367  outputs = keep_outputs(list(make_cell(
368  pred_mh,
369  input,
370  sequence_lens,
371  initial_states_sliced,
372  input_size,
373  hidden_size,
374  name,
375  drop_states=False,
376  forward_only=True,
377  )))
378 
379  if direction_offset == 1:
380  outputs[0] = pred_mh.net.ReversePackedSegs(
381  [outputs[0], seq_lens_for_reverse], name + "/output-reversed")
382 
383  return outputs
384 
385  @classmethod
386  def _create_rnn_variant(cls, init_model, pred_model, n, opset_version):
387  assert init_model is not None, "cannot convert RNNs without access to the full model"
388  assert pred_model is not None, "cannot convert RNNs without access to the full model"
389 
390  attrs = dict(n.attrs) # make a copy, which is safe to mutate
391  hidden_size = attrs.pop('hidden_size')
392  direction = force_unicode(attrs.pop('direction', 'forward'))
393 
394  if n.op_type == 'RNN':
395  activation = force_unicode(attrs.pop('activations', ('tanh',))[0])
396  elif n.op_type == 'GRU':
397  linear_before_reset = attrs.pop('linear_before_reset', 0)
398 
399  assert not attrs, "unsupported RNN attributes: " + str(attrs.keys())
400  assert direction in ['forward', 'bidirectional'], "unsupported backwards RNN/GRU/LSTM"
401 
402  if n.op_type in ['RNN', 'GRU']:
403  input_blob, W, R, B, sequence_lens, initial_h = n.inputs
404  elif n.op_type == 'LSTM':
405  input_blob, W, R, B, sequence_lens, initial_h, initial_c = n.inputs
406 
407  if sequence_lens == "":
408  sequence_lens = None
409 
410  for x in itertools.chain(init_model.graph.input,
411  init_model.graph.value_info,
412  pred_model.graph.input,
413  pred_model.graph.value_info):
414  if x.name == W:
415  input_size = x.type.tensor_type.shape.dim[2].dim_value
416  break
417  else:
418  raise RuntimeError("best-effort shape inference for RNN/GRU/LSTM failed")
419 
420  pred_mh = ModelHelper()
421  init_net = core.Net("init-net")
422 
423  init_net.Reshape(W, [W, cls.dummy_name()], shape=[1,-1,0])
424  init_net.Squeeze(W, W, dims=[0])
425  init_net.Reshape(R, [R, cls.dummy_name()], shape=[1,-1,0])
426  init_net.Squeeze(R, R, dims=[0])
427  init_net.Reshape(B, [B, cls.dummy_name()], shape=[1,-1])
428  init_net.Squeeze(B, B, dims=[0])
429 
430  if n.op_type == 'RNN':
431  def reform(*args):
432  pass
433 
434  def make_cell(*args, **kwargs):
435  return rnn_cell.BasicRNN(*args, activation=activation, **kwargs)
436 
437  def make_rnn(direction_offset):
438  return cls._make_rnn_direction(
439  input_blob, B, W, R, [(initial_h, '/initial_h')], sequence_lens,
440  pred_mh, init_net, input_size, hidden_size, 1, direction_offset,
441  "/i2h_b", "/gates_t_b", "/i2h_w", "/gates_t_w",
442  reform, make_cell, lambda x: x)
443 
444  elif n.op_type == 'GRU':
445  def reform(Bi, Br, W_, R_, name, hidden_size, init_net):
446  # caffe2 has a different order from onnx. We need to rearrange
447  # z r h -> r z h
448  reforms = ((W_, 'i2h_w', True, [(0,-1)]),
449  (R_, 'gate_t_w', False, [(0,-1)]),
450  (Bi, 'i2h_b', True, []),
451  (Br, 'gate_t_b', False, []))
452  cls._rnn_reform_weights(reforms, name, hidden_size, init_net,
453  ['update', 'reset', 'output'], [1, 0, 2])
454 
455  def make_cell(*args, **kwargs):
456  return gru_cell.GRU(*args, linear_before_reset=linear_before_reset, **kwargs)
457 
458  def make_rnn(direction_offset):
459  return cls._make_rnn_direction(
460  input_blob, B, W, R, [(initial_h, '/initial_h')], sequence_lens,
461  pred_mh, init_net, input_size, hidden_size, 3, direction_offset,
462  "_bias_i2h", "_bias_gates", "/i2h_w_pre", "/gates_t_w_pre",
463  reform, make_cell, lambda x: x)
464 
465  elif n.op_type == 'LSTM':
466  def reform(Bi, Br, W_, R_, name, hidden_size, init_net):
467  # caffe2 has a different order from onnx. We need to rearrange
468  # i o f c -> i f o c
469  reforms = ((W_, 'i2h_w', True, [(0, -1)]),
470  (R_, 'gates_t_w', True, [(0, -1)]),
471  (Bi, 'i2h_b' , True, []),
472  (Br, 'gates_t_b', True, []))
473  cls._rnn_reform_weights(reforms, name, hidden_size, init_net,
474  ['input', 'output', 'forget', 'cell'], [0, 2, 1, 3])
475 
476  def make_cell(*args, **kwargs):
477  return rnn_cell.LSTM(*args, **kwargs)
478 
479  def make_rnn(direction_offset):
480  return cls._make_rnn_direction(
481  input_blob, B, W, R, [(initial_h, '/initial_h'), (initial_c, '/initial_c')], sequence_lens,
482  pred_mh, init_net, input_size, hidden_size, 4, direction_offset,
483  "/i2h_b", "/gates_t_b", "/i2h_w", "/gates_t_w",
484  reform, make_cell, lambda x: [x[0], x[1], x[3]])
485 
486  if direction == 'forward':
487  outputs = make_rnn(0)
488 
489  # in the forward case, storage is shared between the
490  # last outputs. We need to decouple them so that the
491  # VariableLengthSequencePadding only mutates
492  # n.outputs[0]
493  for i in range(1, len(outputs)):
494  pred_mh.net.Copy(outputs[i], n.outputs[i])
495 
496  if sequence_lens is not None:
497  pred_mh.net.VariableLengthSequencePadding(
498  [outputs[0], sequence_lens], [outputs[0]])
499  pred_mh.net.ExpandDims([outputs[0]], [n.outputs[0]], dims=[1])
500  elif direction == 'bidirectional':
501  outputs_f = make_rnn(0)
502  outputs_b = make_rnn(1)
503 
504  concatted_output, _ = pred_mh.net.Concat(
505  [outputs_f[0], outputs_b[0]], [cls.dummy_name(), cls.dummy_name()], axis=2)
506  if sequence_lens is not None:
507  pred_mh.net.VariableLengthSequencePadding(
508  [concatted_output, sequence_lens], [concatted_output])
509  reshaped_output, _ = pred_mh.net.Reshape(concatted_output, [cls.dummy_name(), cls.dummy_name()], shape=[0,0,-1,2])
510  pred_mh.net.Transpose(reshaped_output, n.outputs[0], axes=[0,2,1,3])
511  for i in range(1, len(n.outputs)):
512  pred_mh.net.Concat([outputs_f[i], outputs_b[i]],
513  [n.outputs[i], cls.dummy_name()], axis=0)
514 
515  # We want to decide whether to put all of our weight-reshaping
516  # operators in the init net or the predict net. We can put
517  # them in the init net iff the inputs to those operators are
518  # already available, either as graph initializers, or as the
519  # output of other operators in the init net. The latter case
520  # occurs, for example, when exporting from pytorch to onnx.
521  # In most production use, we expect has_initializers to be
522  # true.
523  initializers = {i.name for i in init_model.graph.initializer}
524  outputs = {output for node in init_model.graph.node for output in node.output}
525  has_initializers = all(x in initializers or x in outputs for x in (W, R, B))
526 
527  pred_ops = []
528  init_ops = []
529  (init_ops if has_initializers else pred_ops).extend(init_net.Proto().op)
530  pred_ops.extend(pred_mh.Proto().op)
531 
532  return Caffe2Ops(pred_ops, init_ops, list(pred_mh.Proto().external_input))
533 
534  @classmethod
535  def _create_control_op(cls, init_model, pred_model, n, opset_version):
536  control_inputs = []
537  if '__control_inputs' in n.attrs:
538  control_inputs.extend(n.attrs['__control_inputs'])
539  node = cls._common_onnx_node_to_caffe2_op(init_model, pred_model, n, opset_version)
540  node.control_input.extend(control_inputs)
541  return Caffe2Ops([node], [], [])
542 
543  @classmethod
544  def _remove_ssa(cls, net, remap_dict):
545  for op in net.op:
546  for i, name in enumerate(op.output):
547  if name in remap_dict:
548  op.output[i] = remap_dict[name]
549  for i, out in enumerate(net.external_output):
550  if out in remap_dict:
551  net.external_output[i] = remap_dict[out]
552 
553  @classmethod
554  def _create_if(cls, init_model, pred_model, n, opset_version):
555  ops = cls._create_control_op(init_model, pred_model, n, opset_version)
556  assert ops[0][0].type == 'If'
557  if_op = ops[0][0]
558  then_net = else_net = None
559  control_inputs = []
560  for arg in if_op.arg:
561  if arg.name == 'then_net':
562  then_net = arg.n
563  if arg.name == 'else_net':
564  else_net = arg.n
565  if arg.name == '__control_inputs':
566  control_inputs = arg.strings
567 
568  assert then_net and else_net
569  then_net_outs = then_net.external_output
570  else_net_outs = else_net.external_output
571  op_outputs = if_op.output
572  assert len(then_net_outs) == len(else_net_outs)
573  assert len(else_net_outs) == len(op_outputs)
574 
575  for arg in if_op.arg:
576  if arg.name == 'then_net':
577  arg.n.external_input.extend(control_inputs)
578  if arg.name == 'else_net':
579  arg.n.external_input.extend(control_inputs)
580 
581  return ops
582 
583  @classmethod
584  def _create_loop(cls, init_model, pred_model, n, opset_version):
585  ops = cls._create_control_op(init_model, pred_model, n, opset_version)
586  assert ops[0][0].type == 'ONNXWhile'
587  while_op = ops[0][0]
588  while_op.arg.extend([caffe2.python.utils.MakeArgument('has_trip_count', True)])
589  while_op.arg.extend([caffe2.python.utils.MakeArgument('has_cond', True)])
590  while_op.arg.extend([caffe2.python.utils.MakeArgument('disable_scopes', True)])
591  control_inputs = []
592  for arg in while_op.arg:
593  if arg.name == '__control_inputs':
594  control_inputs = arg.strings
595  num_loop_carried_deps = 0
596  for arg in while_op.arg:
597  if arg.name == 'body':
598  num_loop_carried_deps = len(arg.n.external_input) - 2
599  arg.n.external_input.extend(control_inputs)
600  while_op.arg.extend([
601  caffe2.python.utils.MakeArgument('num_loop_carried_deps',
602  num_loop_carried_deps)
603  ])
604 
605  return ops
606 
607  @classmethod
608  def _substitute_raw_value(cls, tp, raw_values_dict):
609  if tp.HasField('raw_data') and tp.raw_data == bytes(b'__EXTERNAL'):
610  if tp.name not in raw_values_dict:
611  raise RuntimeError('TensorProto for value {} referenced raw data but it was not found!'.format(tp.name))
612  else:
613  tp.raw_data = raw_values_dict[tp.name]
614 
615  @classmethod
616  def _visit_and_substitute_raw_values(cls, nodes, raw_values_dict):
617  for node in nodes:
618  for attr in node.attribute:
619  if attr.HasField('t'):
620  cls._substitute_raw_value(attr.t, raw_values_dict)
621  for t in attr.tensors:
622  cls._substitute_raw_value(t, raw_values_dict)
623  if attr.HasField('g'):
624  cls._visit_and_substitute_raw_values(attr.g.node, raw_values_dict)
625  for g in attr.graphs:
626  cls._visit_and_substitute_raw_values(g.node, raw_values_dict)
627 
628  @classmethod
629  def _external_value_resolution_pass(cls, model, raw_values_dict):
630  for init in model.graph.initializer:
631  cls._substitute_raw_value(init, raw_values_dict)
632 
633  cls._visit_and_substitute_raw_values(model.graph.node, raw_values_dict)
634 
635 
636  @classmethod
637  def _direct_initialize_parameters(cls, initializer, ws, device_option):
638  for tp in initializer:
639  ws.FeedBlob(tp.name, onnx.numpy_helper.to_array(tp), device_option)
640 
641  @classmethod
642  def _direct_initialize_inputs(cls, inputs, initialized, ws, device_option):
643  for value_info in inputs:
644  if value_info.name in initialized:
645  continue
646  shape = list(d.dim_value for d in value_info.type.tensor_type.shape.dim)
647  ws.FeedBlob(
648  value_info.name,
649  np.ones(shape, dtype=onnx.mapping.TENSOR_TYPE_TO_NP_TYPE[value_info.type.tensor_type.elem_type]),
650  device_option)
651 
652  @staticmethod
653  def optimize_onnx(input, init=False, predict=False):
654  passes = ['fuse_consecutive_transposes',
655  'eliminate_nop_transpose',
656  'fuse_transpose_into_gemm',
657  'lift_lexical_references']
658  if init:
659  passes.append('split_init')
660  if predict:
661  passes.append('split_predict')
662  out = onnx.optimizer.optimize(input, passes)
663  return out
664 
665  @classmethod
666  def prepare_zip_archive(cls, file, device='CPU', **kwargs):
667  with zipfile.ZipFile(file, mode='r') as z:
668  with z.open('__MODEL_PROTO', 'r') as f:
669  model = onnx.load(f);
670  blob_names = set(z.namelist()) - set('__MODEL_PROTO')
671  # TODO: make this more efficient
672  raw_values_dict = {}
673  for name in blob_names:
674  with z.open(name, 'r') as blob_file:
675  raw_values_dict[name] = blob_file.read()
676 
677  return cls.prepare(model, device, raw_values_dict=raw_values_dict, **kwargs)
678 
679  @classmethod
680  def prepare(cls, model, device='CPU', raw_values_dict=None, **kwargs):
681  '''
682  For Onnx Caffe2Backend, we require that init_graph don't initialize the actual input of the predict_graph,
683 
684  for example, if "img" is the input blob for the predict_net, we require that in init_graph and in
685  initializer of the predict_graph, "img" is not initalized. We don't have a check for this, since
686  there is no way we can know which blob is the input of the predict_graph.
687  '''
688  if not kwargs.pop('no_check_UNSAFE', False):
689  super(Caffe2Backend, cls).prepare(model, device, **kwargs)
690  opset_version = None
691  for imp in model.opset_import:
692  if not imp.HasField("domain") or imp.domain == "":
693  opset_version = imp.version
694  if imp.version > cls._known_opset_version:
695  warnings.warn("This version of onnx-caffe2 targets ONNX operator set version {}, but the model we are trying to import uses version {}. We will try to import it anyway, but if the model uses operators which had BC-breaking changes in the intervening versions, import will fail.".format(cls._known_opset_version, imp.version))
696  else:
697  warnings.warn("Unrecognized operator set {}".format(imp.domain))
698  if opset_version is None:
699  if model.ir_version >= 0x00000003:
700  raise RuntimeError("Model with IR version >= 3 did not specify ONNX operator set version (onnx-caffe2 requires it)")
701  else:
702  opset_version = 1
703 
704  model = onnx.shape_inference.infer_shapes(model)
705 
706  ws = Workspace()
707  device_option = get_device_option(Device(device))
708 
709  init_net, predict_net = cls._onnx_model_to_caffe2_net(model, device, opset_version, False)
710 
711  if raw_values_dict:
712  cls._external_value_resolution_pass(model, raw_values_dict)
713 
714  # Directly load initializer data into blobs in workspace
716  model.graph.initializer,
717  ws,
718  device_option,
719  )
720 
721  initialized = {init.name for init in model.graph.initializer}
722 
724  model.graph.input,
725  initialized,
726  ws,
727  device_option,
728  )
729 
730  uninitialized = [value_info.name for value_info in model.graph.input if value_info.name not in initialized]
731 
732  retval = Caffe2Rep(init_net, predict_net, ws, uninitialized)
733  return retval
734 
735 
736  @classmethod
737  # TODO: This method needs a refactor for clarity
738  def _onnx_node_to_caffe2_op(cls, init_model, pred_model, node_def, opset_version):
739  cbackend = C.Caffe2Backend(cls._dummy_name)
740  if cbackend.support_onnx_import(node_def.op_type):
741 
742  # extract value infos from pred model (value infos of
743  # node's inputs that are in init model should be all
744  # available in pred model)
745  value_infos = []
746  for name in node_def.input:
747  if pred_model is not None:
748  for vi in itertools.chain(pred_model.graph.input,
749  pred_model.graph.output,
750  pred_model.graph.value_info):
751  if vi.name == name:
752  value_infos.append(vi.SerializeToString())
753 
754  op_strs = cbackend.convert_node(node_def.SerializeToString(), value_infos, opset_version)
755  init_ops = []
756  for s in op_strs[0]:
757  op = caffe2_pb2.OperatorDef()
758  op.ParseFromString(s)
759  init_ops.append(op)
760  ops = []
761  for s in op_strs[1]:
762  op = caffe2_pb2.OperatorDef()
763  op.ParseFromString(s)
764  ops.append(op)
765  return Caffe2Ops(ops, init_ops, [])
766 
767  if node_def.op_type in cls._special_operators:
768  translator = getattr(cls, cls._special_operators[node_def.op_type])
769  else:
770  translator = cls._common_onnx_node_to_caffe2_op
771  ops = translator(init_model, pred_model, OnnxNode(node_def), opset_version)
772  if isinstance(ops, Caffe2Ops):
773  return ops
774  if not isinstance(ops, container_abcs.Iterable):
775  ops = [ops]
776  return Caffe2Ops(ops, [], [])
777 
778  _broadcast_operators = {
779  'Add',
780  'Sub',
781  }
782 
783  @classmethod
784  def _common_onnx_node_to_caffe2_op(cls, init_model, pred_model, onnx_node, opset_version):
785  """
786  This translator performs the basic translation of ONNX nodes into
787  Caffe2 operators. Besides doing a straightforward marshalling from
788  one format to another, it also does these extra things:
789 
790  - Renames operators based on '_renamed_operators'
791  - Renames attributes based on '_global_renamed_attrs' and
792  '_per_op_renamed_attrs'
793 
794  If you're writing a custom translator, consider calling this first,
795  and then fixing things up further.
796  """
797  c2_op = caffe2_pb2.OperatorDef()
798 
799  c2_op.input.extend(onnx_node.inputs)
800  c2_op.output.extend(onnx_node.outputs)
801  c2_op.name = onnx_node.name
802 
803 
804  onnx_op_type = onnx_node.op_type
805  broken_version = cls._broken_operators.get(onnx_op_type, float('Inf'))
806  if broken_version <= opset_version:
807  raise ValueError(
808  "Don't know how to translate op {} in ONNX operator set v{} (I only support prior to v{})".format(onnx_op_type, opset_version, broken_version))
809  c2_op.type = cls._renamed_operators.get(onnx_op_type, onnx_op_type)
810  if not core.IsOperator(c2_op.type):
811  raise ValueError(
812  "Don't know how to translate op {}".format(onnx_op_type))
813 
814  def kmap(k):
815  if (onnx_op_type in cls._per_op_renamed_attrs and
816  k in cls._per_op_renamed_attrs[onnx_op_type]):
817  return cls._per_op_renamed_attrs[onnx_op_type][k]
818  if k in cls._global_renamed_attrs:
819  return cls._global_renamed_attrs[k]
820  return k
821  c2_op.arg.extend(onnx_node.attrs.caffe2(kmap=kmap))
822 
823  if opset_version < 7:
824  # onnx opset 7 and newest caffe2 have adopted full onnx broadcast semantics
825  # so we don't need this hack anymore
826  if c2_op.type in cls._broadcast_operators:
827  already_broadcast = False
828  for arg in c2_op.arg:
829  if arg.name == 'broadcast':
830  already_broadcast = True
831  if not already_broadcast:
832  c2_op.arg.extend([caffe2.python.utils.MakeArgument('broadcast', 1)])
833 
834  return c2_op
835 
836  @staticmethod
837  def _all_names_in_graph(graph):
838  if graph is None:
839  return set()
840 
841  names = set()
842  names.update(value_info.name for value_info in graph.input)
843  names.update(value_info.name for value_info in graph.output)
844  for node in graph.node:
845  names.update(node.input)
846  names.update(node.output)
847  return names
848 
849  @classmethod
850  def _graph_to_net(cls, onnx_graph, opset_version):
851  net = caffe2_pb2.NetDef()
852  for node in onnx_graph.node:
853  try:
854  c2ops = cls._onnx_node_to_caffe2_op(
855  None, None, node, opset_version)
856  except Exception as e:
857  print('ONNX FATAL:', e)
858  continue
859  net.op.extend(c2ops.init_ops)
860  net.op.extend(c2ops.ops)
861  net.external_input.extend(c2ops.interface_blobs)
862  net.external_output.extend(
863  value_info.name for value_info in onnx_graph.output)
864  net.external_input.extend(
865  value_info.name for value_info in onnx_graph.input)
866  return net
867 
868  @classmethod
869  def _onnx_model_to_caffe2_net(cls, onnx_model, device, opset_version, include_initializers):
870  device_option = get_device_option(Device(device))
871 
872  onnx_model = onnx.utils.polish_model(onnx_model)
873  init_model = cls.optimize_onnx(onnx_model, init=True)
874  pred_model = cls.optimize_onnx(onnx_model, predict=True)
875 
876  init_net = caffe2_pb2.NetDef()
877  pred_net = caffe2_pb2.NetDef()
878 
879  init_net.name = onnx_model.graph.name + '_init'
880  pred_net.name = onnx_model.graph.name + '_predict'
881 
882  if include_initializers:
883  init_net.op.extend(cls._create_tensor_filling_op(tp) for tp in onnx_model.graph.initializer)
884 
885  cls._dummy_name.reset(cls._all_names_in_graph(init_model.graph) | cls._all_names_in_graph(pred_model.graph))
886 
887  success = True
888  for net, model in ( (init_net, init_model), (pred_net, pred_model) ):
889  net.device_option.CopyFrom(device_option)
890  for node in model.graph.node:
891  try:
892  c2ops = cls._onnx_node_to_caffe2_op(
893  init_model, pred_model, node, opset_version)
894  except Exception as e:
895  success = False
896  print('ONNX FATAL:', e)
897  continue
898  init_net.op.extend(c2ops.init_ops)
899  net.op.extend(c2ops.ops)
900  net.external_input.extend(c2ops.interface_blobs)
901  net.external_output.extend(
902  value_info.name for value_info in model.graph.output)
903  net.external_input.extend(
904  value_info.name for value_info in model.graph.input)
905 
906  if not success:
907  raise RuntimeError('ONNX conversion failed')
908 
909  return init_net, pred_net
910 
911  # wrapper for backwards compatability
912  @classmethod
913  def onnx_graph_to_caffe2_net(cls, model, device="CPU", opset_version=_known_opset_version):
914  return cls._onnx_model_to_caffe2_net(model, device=device, opset_version=opset_version, include_initializers=True)
915 
916  @classmethod
917  def supports_device(cls, device_str):
918  device = Device(device_str)
919  if device.type == DeviceType.CPU:
920  return True
921  elif core.IsGPUDeviceType(device.type):
922  return workspace.has_gpu_support
923  return False
924 
925  @classmethod
926  def is_compatible(cls, model, device='CPU', **kwargs):
927  if hasattr(super(Caffe2Backend, cls), 'is_compatible') \
928  and callable(super(Caffe2Backend, cls).is_compatible):
929  if not super(Caffe2Backend, cls).is_compatible(model, device, **kwargs):
930  return False
931  # TODO: should have an unspported list of operators, be optimistic for now
932  return True
933 
934 prepare = Caffe2Backend.prepare
935 
936 prepare_zip_archive = Caffe2Backend.prepare_zip_archive
937 
938 run_node = Caffe2Backend.run_node
939 
940 run_model = Caffe2Backend.run_model
941 
942 supports_device = Caffe2Backend.supports_device # noqa
943 
944 is_compatible = Caffe2Backend.is_compatible
def MakeArgument(key, value)
Definition: utils.py:119
def _external_value_resolution_pass(cls, model, raw_values_dict)
Definition: backend.py:629
def _direct_initialize_inputs(cls, inputs, initialized, ws, device_option)
Definition: backend.py:642
def optimize_onnx(input, init=False, predict=False)
Definition: backend.py:653
def _create_tensor_filling_op(cls, onnx_tensor, name=None)
Definition: backend.py:252
def _common_onnx_node_to_caffe2_op(cls, init_model, pred_model, onnx_node, opset_version)
Definition: backend.py:784
def _direct_initialize_parameters(cls, initializer, ws, device_option)
Definition: backend.py:637
def _onnx_model_to_caffe2_net(cls, onnx_model, device, opset_version, include_initializers)
Definition: backend.py:869
def prepare(cls, model, device='CPU', raw_values_dict=None, kwargs)
Definition: backend.py:680
def _onnx_node_to_caffe2_op(cls, init_model, pred_model, node_def, opset_version)
Definition: backend.py:738