converter.py 63 KB

12345678910111213141516171819202122232425262728293031323334353637383940414243444546474849505152535455565758596061626364656667686970717273747576777879808182838485868788899091929394959697989910010110210310410510610710810911011111211311411511611711811912012112212312412512612712812913013113213313413513613713813914014114214314414514614714814915015115215315415515615715815916016116216316416516616716816917017117217317417517617717817918018118218318418518618718818919019119219319419519619719819920020120220320420520620720820921021121221321421521621721821922022122222322422522622722822923023123223323423523623723823924024124224324424524624724824925025125225325425525625725825926026126226326426526626726826927027127227327427527627727827928028128228328428528628728828929029129229329429529629729829930030130230330430530630730830931031131231331431531631731831932032132232332432532632732832933033133233333433533633733833934034134234334434534634734834935035135235335435535635735835936036136236336436536636736836937037137237337437537637737837938038138238338438538638738838939039139239339439539639739839940040140240340440540640740840941041141241341441541641741841942042142242342442542642742842943043143243343443543643743843944044144244344444544644744844945045145245345445545645745845946046146246346446546646746846947047147247347447547647747847948048148248348448548648748848949049149249349449549649749849950050150250350450550650750850951051151251351451551651751851952052152252352452552652752852953053153253353453553653753853954054154254354454554654754854955055155255355455555655755855956056156256356456556656756856957057157257357457557657757857958058158258358458558658758858959059159259359459559659759859960060160260360460560660760860961061161261361461561661761861962062162262362462562662762862963063163263363463563663763863964064164264364464564664764864965065165265365465565665765865966066166266366466566666766866967067167267367467567667767867968068168268368468568668768868969069169269369469569669769869970070170270370470570670770870971071171271371471571671771871972072172272372472572672772872973073173273373473573673773873974074174274374474574674774874975075175275375475575675775875976076176276376476576676776876977077177277377477577677777877978078178278378478578678778878979079179279379479579679779879980080180280380480580680780880981081181281381481581681781881982082182282382482582682782882983083183283383483583683783883984084184284384484584684784884985085185285385485585685785885986086186286386486586686786886987087187287387487587687787887988088188288388488588688788888989089189289389489589689789889990090190290390490590690790890991091191291391491591691791891992092192292392492592692792892993093193293393493593693793893994094194294394494594694794894995095195295395495595695795895996096196296396496596696796896997097197297397497597697797897998098198298398498598698798898999099199299399499599699799899910001001100210031004100510061007100810091010101110121013101410151016101710181019102010211022102310241025102610271028102910301031103210331034103510361037103810391040104110421043104410451046104710481049105010511052105310541055105610571058105910601061106210631064106510661067106810691070107110721073107410751076107710781079108010811082108310841085108610871088108910901091109210931094109510961097109810991100110111021103110411051106110711081109111011111112111311141115111611171118111911201121112211231124112511261127112811291130113111321133113411351136113711381139114011411142114311441145114611471148114911501151115211531154115511561157115811591160116111621163116411651166116711681169117011711172117311741175117611771178117911801181118211831184118511861187118811891190119111921193119411951196119711981199120012011202120312041205120612071208120912101211121212131214121512161217121812191220122112221223122412251226122712281229123012311232123312341235123612371238123912401241124212431244124512461247124812491250125112521253125412551256125712581259126012611262126312641265126612671268126912701271127212731274127512761277127812791280128112821283128412851286128712881289129012911292129312941295129612971298129913001301130213031304130513061307130813091310131113121313131413151316131713181319132013211322132313241325132613271328132913301331133213331334133513361337133813391340134113421343134413451346134713481349135013511352135313541355135613571358135913601361136213631364136513661367136813691370137113721373137413751376137713781379138013811382138313841385138613871388138913901391139213931394139513961397139813991400140114021403140414051406140714081409141014111412141314141415141614171418141914201421142214231424142514261427142814291430143114321433143414351436143714381439144014411442144314441445144614471448144914501451145214531454145514561457145814591460146114621463146414651466146714681469147014711472147314741475147614771478147914801481148214831484148514861487148814891490149114921493149414951496149714981499150015011502150315041505150615071508150915101511151215131514151515161517151815191520152115221523152415251526152715281529153015311532153315341535153615371538153915401541154215431544154515461547154815491550155115521553155415551556155715581559156015611562156315641565156615671568156915701571157215731574157515761577157815791580158115821583158415851586158715881589159015911592159315941595159615971598159916001601160216031604160516061607160816091610
  1. # mypy: allow-untyped-defs
  2. import builtins
  3. import logging
  4. import operator
  5. import typing
  6. import warnings
  7. from collections.abc import Sequence
  8. from contextlib import contextmanager
  9. from typing import Any, Callable, Optional, Union
  10. import torch
  11. import torch.export._trace
  12. from torch import _C
  13. from torch._export.passes.replace_quantized_ops_with_standard_ops_pass import (
  14. replace_quantized_ops_with_standard_ops,
  15. )
  16. from torch.export.dynamic_shapes import _tree_map_with_path, Dim
  17. from torch.export.exported_program import ExportedProgram
  18. from torch.export.graph_signature import (
  19. ConstantArgument,
  20. CustomObjArgument,
  21. InputKind,
  22. InputSpec,
  23. OutputKind,
  24. OutputSpec,
  25. TensorArgument,
  26. )
  27. from torch.fx import subgraph_rewriter
  28. log = logging.getLogger(__name__)
  29. def _get_param_count_list(method_graph, args_params):
  30. param_count_list = []
  31. for input_, arg_params_ in zip(method_graph.inputs(), args_params):
  32. if "PackedParams" in str(input_.type()):
  33. in_vars, _ = torch.jit._flatten(arg_params_)
  34. param_count_list.append(len(in_vars))
  35. else:
  36. param_count_list.append(arg_params_ is not None)
  37. return param_count_list
  38. def _trace_and_get_graph_from_model(model, args):
  39. # A basic sanity check: make sure the state_dict keys are the same
  40. # before and after running the model. Fail fast!
  41. orig_state_dict_keys = torch.jit._unique_state_dict(model).keys()
  42. # Disable Autocast cache because it replaces kernel's weight and bias
  43. # by (undesired) constants.
  44. # No perf impact for when there are reused weights since https://github.com/pytorch/pytorch/pull/85665
  45. prev_autocast_cache_enabled = torch.is_autocast_cache_enabled()
  46. torch.set_autocast_cache_enabled(False)
  47. trace_graph, torch_out, _inputs_states = torch.jit._get_trace_graph(
  48. model,
  49. args,
  50. strict=False,
  51. _force_outplace=False,
  52. _return_inputs_states=True,
  53. )
  54. torch.set_autocast_cache_enabled(prev_autocast_cache_enabled)
  55. if orig_state_dict_keys != torch.jit._unique_state_dict(model).keys():
  56. raise RuntimeError(
  57. "state_dict changed after running the tracer; "
  58. "something weird is happening in your model!"
  59. )
  60. return trace_graph, torch_out
  61. def _create_jit_graph(
  62. model: Union[torch.nn.Module, torch.jit.ScriptFunction], args: Sequence[Any]
  63. ) -> tuple[torch.Graph, list["_C.IValue"], Any, Optional[torch.ScriptModule]]:
  64. if isinstance(model, (torch.jit.ScriptFunction, torch.jit.ScriptModule)):
  65. flattened_args = tuple(torch.jit._flatten(tuple(args))[0])
  66. torch_out = None
  67. if isinstance(model, torch.jit.ScriptModule):
  68. try:
  69. graph = model.forward.graph # type: ignore[attr-defined]
  70. except AttributeError as e:
  71. raise RuntimeError("'forward' method must be a script method") from e
  72. _C._jit_pass_onnx_function_substitution(graph)
  73. freezed_module = _C._freeze_module(
  74. typing.cast(_C.ScriptModule, model._c), preserveParameters=True
  75. )
  76. module, params = _C._jit_onnx_list_model_parameters(freezed_module)
  77. method_graph = module._get_method("forward").graph
  78. args_params = tuple(args) + tuple(params)
  79. param_count_list = _get_param_count_list(method_graph, args_params)
  80. in_vars, _ = torch.jit._flatten(args_params)
  81. graph = _C._propagate_and_assign_input_shapes(
  82. method_graph, tuple(in_vars), param_count_list, False, False
  83. )
  84. return graph, params, torch_out, module
  85. # torch.jit.ScriptFunction
  86. params = []
  87. graph = model.graph
  88. _C._jit_pass_onnx_function_substitution(graph)
  89. param_count_list = _get_param_count_list(graph, args)
  90. graph = _C._propagate_and_assign_input_shapes(
  91. graph, flattened_args, param_count_list, False, False
  92. )
  93. return graph, params, torch_out, None
  94. graph, torch_out = _trace_and_get_graph_from_model(model, args)
  95. _C._jit_pass_onnx_lint(graph)
  96. state_dict = torch.jit._unique_state_dict(model)
  97. params = list(state_dict.values())
  98. graph_inputs = list(graph.inputs())
  99. user_input_num = len(graph_inputs) - len(state_dict)
  100. param_names = list(state_dict.keys())
  101. for i, inp in enumerate(graph_inputs):
  102. if i >= user_input_num:
  103. inp.setDebugName(param_names[i - user_input_num])
  104. _C._jit_pass_onnx_function_substitution(graph)
  105. return graph, params, torch_out, None
  106. def list_add(a, b):
  107. return a + b
  108. def list_append(container, element):
  109. return container + [element]
  110. def execute_subgraph_from_prim_loop(
  111. subgraph, iter_idx, len_loop_local_arguments, *args, **kwargs
  112. ):
  113. """
  114. subgraph: GraphModule from sub-block.
  115. iter_idx: The index of interaction.
  116. len_loop_local_arguments: The number of loop local arguments in args.
  117. """
  118. # Loop local variables. TS graph create those as inputs because their values
  119. # are updated inside the loop.
  120. loop_local_args = args[:len_loop_local_arguments]
  121. # Global variables that are not passed in as inputs to the loop sub-blocks
  122. # but are directly used. Most of time, their values are not updated, but
  123. # the only exception is when there are some operations that perform inplace
  124. # updates.
  125. global_args = args[len_loop_local_arguments:]
  126. return subgraph(*global_args, iter_idx, *loop_local_args, **kwargs)
  127. def inplace_optimize_sym_size_div(gm: torch.fx.GraphModule):
  128. def pattern(im, dim, scale):
  129. sym_size_int = torch.ops.aten.sym_size.int(im, dim)
  130. scalar_tensor = torch.ops.aten.scalar_tensor(sym_size_int)
  131. div_scalar_mode = torch.ops.aten.div.Scalar_mode(
  132. scalar_tensor, scale, rounding_mode="trunc"
  133. )
  134. int_tensor = torch.ops.aten.Int.Tensor(div_scalar_mode)
  135. return int_tensor
  136. def replacement(im, dim, scale):
  137. sym_size_int = torch.ops.aten.sym_size.int(im, dim)
  138. return sym_size_int // scale
  139. subgraph_rewriter.replace_pattern(gm, pattern, replacement)
  140. def is_valid_for_codegen(name):
  141. if len(name) == 0:
  142. raise RuntimeError("Empty argument name for codegen")
  143. if name[0].isdigit():
  144. return False
  145. return True
  146. def normalize_name(name: str, prefix: str = "rename") -> str:
  147. name = name.replace(".", "_")
  148. if is_valid_for_codegen(name):
  149. return name
  150. return f"{prefix}_{name}"
  151. def ir_name_to_func_name(name: str) -> str:
  152. """prim::If -> convert_prim_If"""
  153. name_list = name.split("::")
  154. return "convert_" + "_".join(name_list)
  155. def get_node_as_placeholder_or_get_attr(fx_graph, name, is_top_level_graph):
  156. if is_top_level_graph:
  157. return fx_graph.get_attr(name)
  158. return fx_graph.placeholder(name)
  159. _TORCH_DTYPE_TO_ENUM = {
  160. torch.uint8: 0,
  161. torch.int8: 1,
  162. torch.int16: 2,
  163. torch.int32: 3,
  164. torch.int64: 4,
  165. torch.float16: 5,
  166. torch.float32: 6,
  167. torch.float64: 7,
  168. torch.complex32: 8,
  169. torch.complex64: 9,
  170. torch.complex128: 10,
  171. torch.bool: 11,
  172. torch.qint8: 12,
  173. torch.quint8: 13,
  174. torch.bfloat16: 15,
  175. }
  176. _TORCH_ENUM_TO_DTYPE = {value: key for key, value in _TORCH_DTYPE_TO_ENUM.items()}
  177. def get_dtype_as_int(tensor):
  178. """
  179. prim::dtype has the signature "Tensor a) -> int", where it gets the dtype of
  180. the tensor and returns the integer corresponding to this dtype based on the
  181. enum in ScalarType.h
  182. """
  183. dtype = tensor.dtype
  184. if dtype not in _TORCH_DTYPE_TO_ENUM:
  185. raise RuntimeError(f"Unsupported dtype {dtype}")
  186. return _TORCH_DTYPE_TO_ENUM[dtype]
  187. # Those operators will be automatically populated to a instance method
  188. # of TS2FXGraphConverter with name convert_<namespace>_<opname>().
  189. # Please check __init__ for method population implementations.
  190. kind_to_standard_operators: dict[str, Callable[..., Any]] = {
  191. "prim::max": builtins.max,
  192. "prim::min": builtins.min,
  193. "prim::TupleIndex": operator.getitem,
  194. "aten::__is__": operator.is_,
  195. "aten::__isnot__": operator.is_not,
  196. "aten::__not__": operator.not_,
  197. "aten::__contains__": operator.contains,
  198. "prim::dtype": get_dtype_as_int,
  199. "aten::len": len,
  200. # Mapping from specialized op to its symbolic counterpart.
  201. # They currently do not have any other overrides.
  202. "aten::numel": torch.ops.aten.sym_numel,
  203. "aten::size": torch.ops.aten.sym_size,
  204. "aten::storage_offset": torch.ops.aten.sym_storage_offset,
  205. "aten::stride": torch.ops.aten.sym_stride,
  206. }
  207. def get_ir_value_parent_name_and_attr_name(node):
  208. irv_parent_name, irv_name = node.input().debugName(), node.output().debugName()
  209. attr_name = node.s("name")
  210. return irv_name, irv_parent_name, attr_name
  211. def construct_fqn(ir, ref_map, name_map):
  212. name_list = []
  213. while ir in ref_map:
  214. name_list.append(name_map[ir])
  215. ir = ref_map[ir]
  216. return ".".join(reversed(name_list))
  217. def get_block_to_lifted_attrs(
  218. graph: torch._C.Graph,
  219. ) -> tuple[dict[torch._C.Block, set[str]], dict[str, str]]:
  220. """
  221. Perform two passes to get a mapping of blocks to a set of FQNs of its lifted attributes.
  222. When a graph has control flow, the graph will be divided into multiple blocks. We want to convert
  223. each block to a graph which will be passed into torch.cond. A restriction for torch.cond is that model
  224. parameters/buffers are expected to be lifted as inputs to the subgraphs. Before converting the model,
  225. we will run this pass which will:
  226. 1. Figure out which params/buffers are used within blocks through tracing the GetAttr calls.
  227. 2. Process the graph bottom up to find the lifted attributes of each block by taking the union
  228. of the attributes used in the current block, and the lifted attributes of all its child blocks.
  229. Returns:
  230. A mapping of blocks to a set of FQNs of its lifted attributes, and a
  231. mapping of node names to the FQNs of its lifted attributes.
  232. """
  233. # A map from a block to its expected to be lifted arguments.
  234. blocks_to_lifted_attrs: dict[torch._C.Block, set[str]] = {}
  235. # Reference map stores the input (i.e., src) and output (i.e., dest) IR of a
  236. # GetAttr node. By traversing this reference map, we can figure out the
  237. # full IR aliasing pass and figure out the FQN of an attribute.
  238. # E.g., %2 = GetAttr(linear)[%1] --> node_to_parent_map["%2"] = "%1"
  239. node_to_parent_map: dict[str, str] = {}
  240. # Used for reconstructing the FQN of an attribute based on the reference map.
  241. # In nutshell, for each GetAttr call, GetAttr(input IR, attribute name) -> output IR
  242. # This name map stores which attribute name is called for a src IR --> dest IR action.
  243. # E.g., %2 = GetAttr(linear)[%1] --> node_to_attr_name["%2"] = "linear"
  244. node_to_attr_name: dict[str, str] = {}
  245. def _dfs_get_attr_dependency(entry):
  246. """
  247. First DFS path to construct reference map and name map.
  248. """
  249. for node in entry.nodes():
  250. if node.kind() == "prim::GetAttr":
  251. (
  252. irv_name,
  253. irv_parent_name,
  254. attr_name,
  255. ) = get_ir_value_parent_name_and_attr_name(node)
  256. node_to_parent_map[irv_name] = irv_parent_name
  257. node_to_attr_name[irv_name] = attr_name
  258. for block in node.blocks():
  259. _dfs_get_attr_dependency(block)
  260. def _map_blocks_to_lifted_attrs(entry):
  261. """
  262. Walk the graph in a bottom-up fashion to build the expected to be
  263. lifted arguments for each block.
  264. """
  265. arguments: set[str] = set()
  266. for node in entry.nodes():
  267. for block in node.blocks():
  268. # Recursively build.
  269. arguments = arguments.union(_map_blocks_to_lifted_attrs(block))
  270. if node.kind() == "prim::GetAttr":
  271. irv_name = node.output().debugName()
  272. # Skip for intermediate GetAttr, which will anyway not result a FQN.
  273. # E.g., node_to_parent_name: {"%3": "%2", "%2": "%1"}
  274. # node_to_attr_name: {"%3": "weight", "%2": "linear", "%1": "self"}
  275. # There is only one FQN %3-->%2-->%1: self.linear.weight
  276. # %2-->%1 is not a FQN: self.linear
  277. if irv_name not in set(node_to_parent_map.values()):
  278. arguments.add(
  279. construct_fqn(irv_name, node_to_parent_map, node_to_attr_name)
  280. )
  281. if not isinstance(entry, torch._C.Graph): # Skip the top level.
  282. blocks_to_lifted_attrs[entry] = arguments
  283. return arguments
  284. _dfs_get_attr_dependency(graph)
  285. _map_blocks_to_lifted_attrs(graph)
  286. return blocks_to_lifted_attrs, node_to_attr_name
  287. def get_attribute_fqn_from_ts_node(
  288. name_to_attribute_fqn: dict[str, str], node: torch._C.Node
  289. ) -> str:
  290. def get_attr(name: str):
  291. if name in name_to_attribute_fqn:
  292. return name_to_attribute_fqn[name]
  293. else:
  294. raise ValueError(f"Attribute {name} not found")
  295. if node.kind() == "prim::SetAttr":
  296. input_name = next(node.inputs()).debugName()
  297. elif node.kind() == "prim::GetAttr":
  298. input_name = node.input().debugName()
  299. else:
  300. raise RuntimeError(
  301. f"Unexpected node kind when getting attribute fqn. node: {node} "
  302. )
  303. attr_name = node.s("name")
  304. root_attr_name = get_attr(input_name)
  305. attr_fqn = f"{root_attr_name}.{attr_name}" if root_attr_name else attr_name
  306. return attr_fqn
  307. def get_op_overload(node: torch._C.Node):
  308. schema_str = node.schema()
  309. assert schema_str != "(no schema)", f"got empty schema for {node}"
  310. schema: torch._C.FunctionSchema = torch._C.parse_schema(schema_str)
  311. ns, op_name = str(schema.name).split("::")
  312. override = schema.overload_name
  313. try:
  314. op_overload_mod = getattr(torch.ops, ns)
  315. op_overload_packet = getattr(op_overload_mod, op_name)
  316. if override:
  317. op_overload = getattr(op_overload_packet, override)
  318. else:
  319. op_overload = op_overload_packet.default
  320. except Exception as e:
  321. raise RuntimeError(
  322. f"Unable to find operator {node.kind()} with schema {node.schema()}"
  323. ) from e
  324. return op_overload
  325. class TS2FXGraphConverter:
  326. def __init__(
  327. self,
  328. ts_graph: Union[torch._C.Graph, torch._C.Block],
  329. name_to_param: dict[str, torch.Tensor],
  330. name_to_buffer: dict[str, torch.Tensor],
  331. blocks_to_lifted_attrs: dict[torch._C.Block, set[str]],
  332. name_to_non_tensor_attribute: dict[str, Any],
  333. name_to_constant: dict[str, Any],
  334. name_to_attribute_fqn: dict[str, str],
  335. ):
  336. self.ts_graph = ts_graph
  337. # Mapping of parameter FQN to actual parameter value
  338. self.name_to_param = name_to_param
  339. # Mapping of buffer FQN to actual buffer value
  340. self.name_to_buffer = name_to_buffer
  341. self.fx_graph: torch.fx.Graph = torch.fx.Graph()
  342. self.input_specs: list[InputSpec] = []
  343. self.output_specs: list[OutputSpec] = []
  344. # Mapping of TS node name to converted FX node
  345. self.name_to_node: dict[
  346. str, Union[torch.fx.Node, list[torch.fx.Node], dict[Any, torch.fx.Node]]
  347. ] = {}
  348. # Mapping of TS node name to constant value (int, str, TorchBind obj,
  349. # tensor constants ...)
  350. self.name_to_constant: dict[str, Any] = name_to_constant
  351. # Mapping from torchscript node output name to attribute fully qualified name
  352. self.name_to_attribute_fqn: dict[str, str] = name_to_attribute_fqn
  353. # Mapping from fully qualified name to real values or a fx graph node
  354. # During convert, this represents the current value of a non-tensor attribute
  355. # One use case is:
  356. # def forward(self, x):
  357. # c1 = self.count
  358. # self.count += 1
  359. # c2 = self.count
  360. # return x + c1 + c2
  361. self.name_to_non_tensor_attribute_node: dict[str, Any] = {}
  362. # Mapping from fully qualified name to initial real values inputs
  363. # We separate it from self.name_to_non_tensor_attribute_node since
  364. # we need initial real value input when we construct fx.GraphModule
  365. self.name_to_non_tensor_attribute: dict[str, Any] = name_to_non_tensor_attribute
  366. self.subgraphs: dict[str, torch.fx.GraphModule] = {}
  367. # Mapping of block to list of attributes that need to be lifted for each
  368. # block
  369. self.blocks_to_lifted_attrs = blocks_to_lifted_attrs
  370. # Populate methods for the standard operators.
  371. for k in kind_to_standard_operators.keys():
  372. handler_func_name = ir_name_to_func_name(k)
  373. # Create an indirect function call:
  374. # convert_<namespace>_<opname> --> lambda node: _convert_standard_operator(node)
  375. setattr(
  376. self,
  377. handler_func_name,
  378. lambda node: self._convert_standard_operators(node),
  379. )
  380. # This stores a list of return results that do not appear in the original TS
  381. # graph's outputs. The reason we maintain this is because some operations in the sub-block
  382. # might have inplace updates to the variable defined in the parent fx graph. After
  383. # the execution of that sub-block, the variable defined in the parent fx graph also
  384. # needs to be updated.
  385. self.name_update_from_subblock_to_parent: set[str] = set()
  386. def _is_get_attr_node(self, fqn):
  387. return (
  388. fqn in self.name_to_buffer
  389. or fqn in self.name_to_param
  390. or (
  391. fqn in self.name_to_constant
  392. and isinstance(self.name_to_constant[fqn], torch.ScriptObject)
  393. )
  394. )
  395. def _convert_block_to_subgraph(self, node: torch._C.Node, arguments: list[str]):
  396. subgraph_nodes, subgraph_converters = [], []
  397. for block in node.blocks():
  398. subgraph_converter = TS2FXGraphConverter(
  399. block,
  400. self.name_to_param,
  401. self.name_to_buffer,
  402. self.blocks_to_lifted_attrs,
  403. {},
  404. self.name_to_constant,
  405. self.name_to_attribute_fqn,
  406. )
  407. for block_arg in arguments:
  408. normalized_block_arg_name = normalize_name(block_arg)
  409. placeholder_node = subgraph_converter.fx_graph.placeholder(
  410. normalized_block_arg_name
  411. )
  412. subgraph_converter.name_to_node[block_arg] = placeholder_node
  413. subgraph = subgraph_converter.convert()
  414. subgraph_name = self.add_subgraph(subgraph)
  415. subgraph_nodes.append(self.fx_graph.get_attr(subgraph_name))
  416. subgraph_converters.append(subgraph_converter)
  417. return subgraph_nodes, subgraph_converters
  418. def _identify_inputs_as_arguments(self, entry):
  419. """
  420. Identify inputs from the innermost sub-block. This is needed
  421. for nested sub-blocks when the input is hidden in the nested sub-block.
  422. E.g., example IR of input is hidden in the nested sub-block.
  423. Graph[x.1]
  424. %1 = ...
  425. Block[]
  426. Block[x.1]
  427. %2 = x.1 ...
  428. """
  429. arguments: set[str] = set()
  430. for block in entry.blocks():
  431. for block_node in block.nodes():
  432. for block_node_in in block_node.inputs():
  433. if (
  434. block_node_in.debugName() in self.name_to_node
  435. and block_node_in.debugName() not in self.name_to_attribute_fqn
  436. ):
  437. arguments.add(block_node_in.debugName())
  438. arguments = arguments.union(
  439. self._identify_inputs_as_arguments(block_node)
  440. )
  441. return arguments
  442. def is_top_level_graph(self):
  443. return isinstance(self.ts_graph, torch._C.Graph)
  444. def add_subgraph(self, subgraph) -> str:
  445. name = f"subgraph_{len(self.subgraphs)}"
  446. self.subgraphs[name] = subgraph
  447. return name
  448. def get_args_kwargs(self, node: torch._C.Node, schema):
  449. args = []
  450. kwargs = {}
  451. for input, schema_arg in zip(node.inputs(), schema.arguments):
  452. if schema_arg.kwarg_only:
  453. kwargs[schema_arg.name] = self.get_fx_value_by_ir_value(input)
  454. else:
  455. args.append(self.get_fx_value_by_ir_value(input))
  456. return tuple(args), kwargs
  457. def get_fx_value_by_ir_value(self, value: torch._C.Value):
  458. value_name = value.debugName()
  459. if value_name in self.name_to_node:
  460. input_node = self.name_to_node[value_name]
  461. return input_node
  462. elif value_name in self.name_to_constant:
  463. if isinstance(self.name_to_constant[value_name], torch.ScriptObject):
  464. return self.fx_graph.get_attr(value_name)
  465. return self.name_to_constant[value_name]
  466. elif value_name in self.name_to_attribute_fqn:
  467. return self.get_fx_value_by_fqn(self.name_to_attribute_fqn[value_name])
  468. else:
  469. raise ValueError(f"Input {value_name} not found")
  470. def get_fx_value_by_fqn(self, name):
  471. if name in self.name_to_node:
  472. fx_node = self.name_to_node[name]
  473. elif name in self.name_to_constant:
  474. fx_node = self.name_to_constant[name]
  475. elif name in self.name_to_non_tensor_attribute_node:
  476. fx_node = self.name_to_non_tensor_attribute_node[name]
  477. elif name in self.name_to_non_tensor_attribute:
  478. fx_node = self.name_to_non_tensor_attribute[name]
  479. else:
  480. raise ValueError(f"Attribute {name} not found")
  481. return fx_node
  482. def convert(self) -> torch.fx.GraphModule:
  483. self.convert_graph_inputs()
  484. for node in self.ts_graph.nodes():
  485. self.convert_node(node)
  486. self.convert_graph_outputs()
  487. # Pass parameter and buffer to the root for lookup.
  488. gm = torch.fx.GraphModule(
  489. {
  490. **self.subgraphs,
  491. **self.name_to_param,
  492. **self.name_to_buffer,
  493. **self.name_to_non_tensor_attribute,
  494. **self.name_to_constant,
  495. },
  496. self.fx_graph,
  497. )
  498. inplace_optimize_sym_size_div(gm)
  499. gm.graph.lint()
  500. return gm
  501. def convert_graph_inputs(self):
  502. for graph_input in self.ts_graph.inputs():
  503. name = graph_input.debugName()
  504. if name in self.name_to_param:
  505. normalized_name = normalize_name(name)
  506. self.input_specs.append(
  507. InputSpec(
  508. InputKind.PARAMETER,
  509. arg=TensorArgument(name=normalized_name),
  510. target=name,
  511. )
  512. )
  513. fx_node = get_node_as_placeholder_or_get_attr(
  514. self.fx_graph, name, self.is_top_level_graph()
  515. )
  516. elif name in self.name_to_buffer:
  517. normalized_name = normalize_name(name)
  518. self.input_specs.append(
  519. InputSpec(
  520. InputKind.BUFFER,
  521. arg=TensorArgument(name=normalized_name),
  522. target=name,
  523. persistent=True,
  524. )
  525. )
  526. fx_node = get_node_as_placeholder_or_get_attr(
  527. self.fx_graph, name, self.is_top_level_graph()
  528. )
  529. elif name in self.name_to_constant:
  530. assert isinstance(self.name_to_constant[name], torch.ScriptObject), (
  531. "Input conversion only handles ScriptObject"
  532. )
  533. normalized_name = normalize_name(name)
  534. self.input_specs.append(
  535. InputSpec(
  536. InputKind.CUSTOM_OBJ,
  537. arg=CustomObjArgument(
  538. name=normalized_name, class_fqn=normalized_name
  539. ),
  540. target=name,
  541. persistent=False,
  542. )
  543. )
  544. fx_node = get_node_as_placeholder_or_get_attr(
  545. self.fx_graph, name, self.is_top_level_graph()
  546. )
  547. elif isinstance(graph_input.type(), torch.ClassType):
  548. # Directly skip inputs that are ScriptObject but not used in the graph.
  549. continue
  550. else:
  551. normalized_name = normalize_name(name, prefix="input")
  552. self.input_specs.append(
  553. InputSpec(
  554. InputKind.USER_INPUT,
  555. arg=TensorArgument(name=normalized_name),
  556. target=name,
  557. )
  558. )
  559. fx_node = self.fx_graph.placeholder(normalized_name)
  560. self.name_to_node[name] = fx_node
  561. def convert_aten_Float(self, node: torch._C.Node):
  562. def to_float_tensor(t):
  563. return t.to(dtype=torch.float).item()
  564. inp_list = [self.get_fx_value_by_ir_value(inp) for inp in node.inputs()] # noqa: C416
  565. fx_node = self.fx_graph.call_function(
  566. to_float_tensor,
  567. tuple(inp_list),
  568. )
  569. self.name_to_node[node.output().debugName()] = fx_node
  570. def convert_aten_tensor(self, node: torch._C.Node):
  571. """aten::tensor creates a constant tensor ad-hoc --> GetAttr"""
  572. args, kwargs = self.get_args_kwargs(node, torch.ops.aten.tensor.default._schema)
  573. for k in kwargs:
  574. if k == "requires_grad":
  575. kwargs[k] = bool(kwargs[k]) # 0 -> False, 1 -> True
  576. to_tensor = (
  577. torch.tensor
  578. if all(isinstance(a, int) for a in args)
  579. else torch._refs.tensor
  580. )
  581. def target(*args, **kwargs):
  582. if "dtype" in kwargs and kwargs["dtype"] is not None:
  583. kwargs["dtype"] = _TORCH_ENUM_TO_DTYPE[kwargs["dtype"]]
  584. return to_tensor(*args, **kwargs)
  585. # def to_dynamic_tensor(*args, **kwargs):
  586. # if "dtype" in kwargs and kwargs["dtype"] is not None:
  587. # kwargs["dtype"] = _TORCH_ENUM_TO_DTYPE[kwargs["dtype"]]
  588. # return torch._refs.tensor(*args, **kwargs)
  589. output_name = node.output().debugName()
  590. fx_node = self.fx_graph.call_function(target, args, kwargs)
  591. self.name_to_node[output_name] = fx_node
  592. def convert_aten_append(self, node: torch._C.Node):
  593. # special handle python list append: "aten::append.t(t[](a!) self, t(c -> *) el) -> t[](a!)"
  594. # inplace append to the list!! This is kinda crazy, as we are inplace mutating the list
  595. # This makes the converter "non-functional", and the result depends on the order of the nodes being converter
  596. # In a sense, the converter now becomes an stateful interpreter
  597. warnings.warn(
  598. "Converting aten::append.t, which is a inplace mutation of the list. "
  599. "This makes the converter non-functional: the result depends on the order of the append nodes being converter!"
  600. )
  601. args = tuple(self.get_fx_value_by_ir_value(inp) for inp in node.inputs())
  602. fx_node = self.fx_graph.call_function(list_append, args)
  603. self.name_to_node[node.output().debugName()] = fx_node
  604. # inplace mutate arg[0], which is the python list
  605. self.name_to_node[node.inputsAt(0).debugName()] = fx_node
  606. # Variables that need to be updated to parent module.
  607. if not self.is_top_level_graph() and args[0].op == "placeholder":
  608. self.name_update_from_subblock_to_parent.add(node.inputsAt(0).debugName())
  609. def convert_prim_Constant(self, node: torch._C.Node):
  610. name = node.output().debugName()
  611. value: Any = None
  612. if node.hasAttribute("value"):
  613. constant_kind = node.kindOf("value")
  614. if constant_kind == "i":
  615. value = node.i("value")
  616. elif constant_kind == "f":
  617. value = node.f("value")
  618. elif constant_kind == "s":
  619. value = node.s("value")
  620. elif constant_kind == "t":
  621. alias_name = (
  622. f"lifted_tensor_{name}" # Follow naming convention from EP tracing.
  623. )
  624. fx_node = self.fx_graph.get_attr(alias_name)
  625. self.name_to_node[name] = fx_node
  626. name, value = alias_name, node.t("value")
  627. elif constant_kind == "ival":
  628. value = node.ival("value")
  629. else:
  630. raise ValueError(f"Unsupported constant type: {node.kindOf('value')}")
  631. else:
  632. value = None
  633. self.name_to_constant[name] = value
  634. def convert_prim_CallMethod(self, node: torch._C.Node):
  635. inp_list = [self.get_fx_value_by_ir_value(inp) for inp in node.inputs()] # noqa: C416
  636. fx_node = self.fx_graph.call_method(
  637. node.s("name"),
  638. tuple(inp_list),
  639. )
  640. self.name_to_node[node.output().debugName()] = fx_node
  641. def convert_prim_device(self, node: torch._C.Node):
  642. input_type = node.input().type()
  643. if input_type.isSubtypeOf(torch._C.TensorType.get()):
  644. device = input_type.device() # type: ignore[attr-defined]
  645. output_name = node.output().debugName()
  646. self.name_to_constant[output_name] = device
  647. else:
  648. raise ValueError(f"Unsupported JitType ({input_type}) when get device")
  649. def convert_prim_GetAttr(self, node: torch._C.Node):
  650. # Build fully qulified name
  651. attr_fqn = get_attribute_fqn_from_ts_node(self.name_to_attribute_fqn, node)
  652. output_name = node.output().debugName()
  653. self.name_to_attribute_fqn[output_name] = attr_fqn
  654. if self.is_top_level_graph():
  655. if self._is_get_attr_node(attr_fqn):
  656. # We insert a get_attr node due to two reasons.
  657. # First, ts graph does not lift tensor constants as input nodes. So
  658. # tensor constants may be ignored by in convert_graph_inputs().
  659. # Second, attr_fqn may have been written to via SetAttr. Two
  660. # GetAttr may give different values.
  661. self.name_to_node[output_name] = self.fx_graph.get_attr(attr_fqn)
  662. else:
  663. if attr_fqn not in self.name_to_non_tensor_attribute_node:
  664. self.name_to_non_tensor_attribute_node[attr_fqn] = (
  665. self.name_to_non_tensor_attribute[attr_fqn]
  666. )
  667. self.name_to_node[output_name] = self.name_to_non_tensor_attribute_node[
  668. attr_fqn
  669. ]
  670. else:
  671. # Special support for if blocks which do not allow SetAttr TorchScript
  672. # node and get_attr FX Graph Node.
  673. if self._is_get_attr_node(attr_fqn):
  674. self.name_to_node[output_name] = self.name_to_node[attr_fqn]
  675. def convert_prim_SetAttr(self, node: torch._C.Node):
  676. attr_fqn = get_attribute_fqn_from_ts_node(self.name_to_attribute_fqn, node)
  677. attr_value = tuple(node.inputs())[1]
  678. ts_graph_tensor_input = self.get_fx_value_by_ir_value(attr_value)
  679. if self._is_get_attr_node(attr_fqn):
  680. fx_attr_node = self.fx_graph.get_attr(attr_fqn)
  681. self.fx_graph.call_function(
  682. torch.Tensor.copy_, (fx_attr_node, ts_graph_tensor_input)
  683. )
  684. else:
  685. self.name_to_non_tensor_attribute_node[attr_fqn] = ts_graph_tensor_input
  686. def convert_call_function_op(self, node: torch._C.Node):
  687. target = get_op_overload(node)
  688. args, kwargs = self.get_args_kwargs(node, target._schema)
  689. fx_node = self.fx_graph.call_function(target, args, kwargs)
  690. # TODO: convert sourceRange() into stack_trace
  691. # fx_node.meta["stack_trace"] = node.sourceRange()
  692. if node.outputsSize() == 1:
  693. output_name = node.output().debugName()
  694. self.name_to_node[output_name] = fx_node
  695. else:
  696. for i, outp in enumerate(node.outputs()):
  697. output_name = outp.debugName()
  698. next_fx_node = self.fx_graph.call_function(
  699. operator.getitem, (fx_node, i)
  700. )
  701. self.name_to_node[output_name] = next_fx_node
  702. def convert_prim_TupleConstruct(self, node: torch._C.Node):
  703. self._convert_prim_iterator(node)
  704. def convert_prim_ListConstruct(self, node: torch._C.Node):
  705. self._convert_prim_iterator(node)
  706. def _convert_prim_iterator(self, node: torch._C.Node):
  707. output_list = [self.get_fx_value_by_ir_value(inp) for inp in node.inputs()]
  708. output_name = node.output().debugName()
  709. self.name_to_node[output_name] = output_list
  710. def convert_prim_DictConstruct(self, node: torch._C.Node):
  711. output_dict = {}
  712. k, v = None, None
  713. for i, inp in enumerate(node.inputs()):
  714. # We assume key value are stored in pair in the DictConstruct.
  715. # The first element is the key and the following is the value.
  716. if i % 2 == 0:
  717. k = self.get_fx_value_by_ir_value(inp)
  718. else:
  719. v = self.get_fx_value_by_ir_value(inp)
  720. assert k is not None and v is not None, (
  721. "DictConstruct has an empty key value pair."
  722. )
  723. output_dict[k] = v
  724. k, v = None, None
  725. assert k is None and v is None, (
  726. "DictConstruct has an odd number of elements (violating our assumption)."
  727. )
  728. output_name = node.output().debugName()
  729. self.name_to_node[output_name] = output_dict
  730. def convert_prim_ListUnpack(self, node: torch._C.Node):
  731. self._convert_prim_unpack_iterator(node)
  732. def convert_prim_TupleUnpack(self, node: torch._C.Node):
  733. self._convert_prim_unpack_iterator(node)
  734. def _convert_prim_unpack_iterator(self, node: torch._C.Node):
  735. # Single input and multiple outputs for unpacking.
  736. for i, outp in enumerate(node.outputs()):
  737. outp_name = outp.debugName()
  738. inp = self.get_fx_value_by_ir_value(node.input())
  739. fx_node = self.fx_graph.call_function(operator.getitem, (inp, i))
  740. self.name_to_node[outp_name] = fx_node
  741. def convert_aten_Int(self, node: torch._C.Node):
  742. # converts aten::Int as aten._to_copy + aten::_local_scalar_dense
  743. target = torch.ops.aten._to_copy.default
  744. args = tuple(self.get_fx_value_by_ir_value(input) for input in node.inputs())
  745. to_copy_node = self.fx_graph.call_function(target, args, {"dtype": torch.int32})
  746. fx_node = self.fx_graph.call_function(
  747. torch.ops.aten._local_scalar_dense.default, (to_copy_node,)
  748. )
  749. # TODO: convert sourceRange() into stack_trace
  750. # fx_node.meta["stack_trace"] = node.sourceRange()
  751. output_name = node.output().debugName()
  752. self.name_to_node[output_name] = fx_node
  753. def convert_prim_NumToTensor(self, node: torch._C.Node):
  754. # Converts prim::NumToTensor as aten.scalar_tensor.
  755. # prim::NumToTensor IRs are currently triggered by:
  756. # .size() https://github.com/pytorch/pytorch/blob/main/torch/csrc/jit/frontend/tracer.cpp#L950
  757. # .numel() https://github.com/pytorch/pytorch/blob/main/torch/csrc/jit/frontend/tracer.cpp#L971
  758. # For both of those APIs, torch.jit.trace implicitly sets the output tensor type
  759. # to be LongTensor.
  760. target = torch.ops.aten.scalar_tensor
  761. args = tuple(self.get_fx_value_by_ir_value(input) for input in node.inputs())
  762. fx_node = self.fx_graph.call_function(target, args, {"dtype": torch.long})
  763. output_name = node.output().debugName()
  764. self.name_to_node[output_name] = fx_node
  765. def convert_prim_CreateObject(self, node: torch._C.Node):
  766. output_name = node.output().debugName()
  767. self.name_to_attribute_fqn[output_name] = ""
  768. def convert_aten__convolution(self, node: torch._C.Node):
  769. # converts aten::_convolution as aten.convolution, since aten::_convolution
  770. # doesn't have a meta function
  771. target = torch.ops.aten.convolution.default
  772. args, kwargs = self.get_args_kwargs(node, target._schema)
  773. fx_node = self.fx_graph.call_function(target, args, kwargs)
  774. output_name = node.output().debugName()
  775. self.name_to_node[output_name] = fx_node
  776. def convert_aten_div(self, node: torch._C.Node):
  777. target = get_op_overload(node)
  778. schema = target._schema
  779. args, kwargs = self.get_args_kwargs(node, schema)
  780. # converts aten::div.Tensor_mode(x, tensor_constant)
  781. # as aten.div.Scalar_mode(x, tensor_constant.item())
  782. if schema.overload_name == "Tensor_mode":
  783. arg1_name = args[1].name
  784. if arg1_name in self.name_to_constant and isinstance(
  785. self.name_to_constant[arg1_name], torch.Tensor
  786. ):
  787. tensor_constant = self.name_to_constant[arg1_name]
  788. if tensor_constant.numel() == 1:
  789. updated_args = list(args)
  790. updated_args[1] = self.name_to_constant[arg1_name].item()
  791. fx_node = self.fx_graph.call_function(
  792. torch.ops.aten.div.Scalar_mode,
  793. tuple(updated_args),
  794. kwargs,
  795. )
  796. # TODO: convert sourceRange() into stack_trace
  797. # fx_node.meta["stack_trace"] = node.sourceRange()
  798. output_name = node.output().debugName()
  799. self.name_to_node[output_name] = fx_node
  800. return
  801. self.convert_call_function_op(node)
  802. def convert_aten___getitem__(self, node: torch._C.Node):
  803. input_container, index = tuple(
  804. self.get_fx_value_by_ir_value(input) for input in node.inputs()
  805. )
  806. fx_node = self.fx_graph.call_function(
  807. operator.getitem, (input_container, index)
  808. )
  809. output_name = node.output().debugName()
  810. self.name_to_node[output_name] = fx_node
  811. def convert_aten_to(self, node: torch._C.Node):
  812. target = get_op_overload(node)
  813. args, _kwargs = self.get_args_kwargs(node, target._schema)
  814. # special handle aten.to.dtype and aten.to.prim_dtype followed by inplace_mutation_op
  815. # coz aten.to + inplace_mutation_op pattern would trigger
  816. # "cannot mutate tensors with frozen storage" functionalization error.
  817. # To work around the issue, we override the copy to be True, so that the output
  818. # is for sure not an alias of input
  819. if target == torch.ops.aten.to.dtype or target == torch.ops.aten.to.prim_dtype:
  820. user_nodes = [use.user for use in node.output().uses()]
  821. user_targets = [
  822. get_op_overload(user_node)
  823. for user_node in user_nodes
  824. if user_node.schema() != "(no schema)"
  825. ]
  826. has_mutable_target = any(
  827. target._schema.is_mutable for target in user_targets
  828. )
  829. if has_mutable_target:
  830. assert len(args) >= 4
  831. new_args = list(args)
  832. new_args[3] = True # copy, override to True
  833. fx_node = self.fx_graph.call_function(
  834. torch.ops.aten.to.dtype, tuple(new_args)
  835. )
  836. # temp hack to work around the issue https://github.com/pytorch/pytorch/issues/131679
  837. # When this issue is fixed, the clone node would be no longer needed
  838. clone_node = self.fx_graph.call_function(
  839. torch.ops.aten.clone.default, (fx_node,)
  840. )
  841. output_name = node.output().debugName()
  842. self.name_to_node[output_name] = clone_node
  843. return
  844. self.convert_call_function_op(node)
  845. def convert_aten_add(self, node: torch._C.Node):
  846. if node.schema() == "(no schema)":
  847. if isinstance(node.inputsAt(0).type(), torch.ListType) and isinstance(
  848. node.inputsAt(1).type(), torch.ListType
  849. ):
  850. target = torch.ops.aten.add.t
  851. else:
  852. raise RuntimeError(f"unable to determined the target for {node}")
  853. else:
  854. target = get_op_overload(node)
  855. if target == torch.ops.aten.add.t:
  856. # special handle python list/tuple add: "aten::add.t(t[] a, t[] b) -> t[]" for
  857. # RuntimeError: aten::add() Expected a value of type 'List[t]' for argument 'a' but instead found type 'immutable_list'.
  858. args, _kwargs = self.get_args_kwargs(node, target._schema)
  859. output_name = node.output().debugName()
  860. self.name_to_node[output_name] = self.fx_graph.call_function(list_add, args)
  861. else:
  862. self.convert_call_function_op(node)
  863. def _check_prim_loop_support(self, node):
  864. inputs = list(node.inputs())
  865. # TODO: (1/N) stage.
  866. if inputs[0].debugName() not in self.name_to_constant:
  867. raise RuntimeError(
  868. "prim::Loop currently cannot run with dynamic value of number of iterations."
  869. )
  870. # Make sure the condition is not updated in the subblock.
  871. subblock = next(node.blocks())
  872. condition_output_name = next(subblock.outputs()).debugName()
  873. for node in subblock.nodes():
  874. if (
  875. node.outputsSize() == 1
  876. and node.output().debugName() == condition_output_name
  877. ):
  878. raise RuntimeError(
  879. "prim::Loop currently cannot run with dynamic value of condition."
  880. )
  881. if node.outputsSize() >= 2:
  882. for outp in node.outputs():
  883. if outp.debugName() == condition_output_name:
  884. raise RuntimeError(
  885. "prim::Loop currently cannot run with dynamic value of condition."
  886. )
  887. def convert_prim_Loop(self, node: torch._C.Node):
  888. inputs = list(node.inputs())
  889. self._check_prim_loop_support(node)
  890. num_iterations = self.get_fx_value_by_ir_value(inputs[0])
  891. # Find inputs.
  892. loop_local_arguments = [inp.debugName() for inp in inputs[2:]]
  893. global_arguments = self._identify_inputs_as_arguments(node)
  894. # Lift parameters as inputs.
  895. for block in node.blocks():
  896. global_arguments = global_arguments.union(
  897. self.blocks_to_lifted_attrs[block]
  898. )
  899. global_arguments = list(global_arguments)
  900. subgraph_nodes, subgraph_converters = self._convert_block_to_subgraph(
  901. node, global_arguments
  902. )
  903. assert len(subgraph_nodes) == 1
  904. subgraph_converter = subgraph_converters[0]
  905. if not self.is_top_level_graph():
  906. self.name_update_from_subblock_to_parent = (
  907. self.name_update_from_subblock_to_parent.union(
  908. subgraph_converter.name_update_from_subblock_to_parent
  909. )
  910. )
  911. fx_block_args = [
  912. self.get_fx_value_by_fqn(name)
  913. for name in loop_local_arguments + global_arguments
  914. ]
  915. for iter_idx in range(num_iterations):
  916. loop_node = self.fx_graph.call_function(
  917. execute_subgraph_from_prim_loop,
  918. # Check execute_node function for the expected arguments order.
  919. (
  920. subgraph_nodes[0],
  921. iter_idx,
  922. len(loop_local_arguments),
  923. *fx_block_args,
  924. ),
  925. {},
  926. )
  927. # Update the value of loop local variables.
  928. if node.outputsSize() >= 1:
  929. for i, outp in enumerate(node.outputs()):
  930. output_name = outp.debugName()
  931. self.name_to_node[output_name] = self.fx_graph.call_function(
  932. operator.getitem,
  933. (
  934. loop_node,
  935. i + 1,
  936. ), # + 1 because the 0th element is the condition.
  937. )
  938. fx_block_args[i] = self.name_to_node[output_name]
  939. # Update the value of global variables, whose values are modified inplace.
  940. for i, name in enumerate(
  941. subgraph_converter.name_update_from_subblock_to_parent
  942. ):
  943. self.name_to_node[name] = self.fx_graph.call_function(
  944. operator.getitem,
  945. (
  946. loop_node,
  947. i + node.outputsSize() + 1,
  948. ), # + 1 because the 0th element is the condition.
  949. )
  950. global_argument_index = global_arguments.index(name)
  951. fx_block_args[i + node.outputsSize() + global_argument_index] = (
  952. self.name_to_node[name]
  953. )
  954. def _check_set_attr_in_if_block(self, if_node: torch._C.Node):
  955. for block in if_node.blocks():
  956. for node in block.nodes():
  957. if node.kind() == "prim::SetAttr":
  958. raise RuntimeError(
  959. "During converting prim::If to torch.cond, found prim::SetAttr op"
  960. " which is not supported yet. Please file an issue if you come "
  961. "across this error."
  962. )
  963. def convert_prim_If(self, node: torch._C.Node):
  964. self._check_set_attr_in_if_block(node)
  965. inputs = list(node.inputs())
  966. assert len(inputs) == 1
  967. predicate = self.get_fx_value_by_ir_value(inputs[0])
  968. # Find inputs.
  969. arguments = self._identify_inputs_as_arguments(node)
  970. # Lift parameters as inputs.
  971. for block in node.blocks():
  972. arguments = arguments.union(self.blocks_to_lifted_attrs[block])
  973. arguments = list(arguments)
  974. subgraph_nodes, _ = self._convert_block_to_subgraph(node, arguments)
  975. assert len(subgraph_nodes) == 2
  976. fx_block_args = [self.get_fx_value_by_fqn(name) for name in arguments]
  977. args = (
  978. predicate,
  979. subgraph_nodes[0],
  980. subgraph_nodes[1],
  981. tuple(fx_block_args),
  982. )
  983. cond_node = self.fx_graph.call_function(torch.cond, args, {})
  984. # prim::If can also have zero output.
  985. if node.outputsSize() == 1:
  986. output_name = node.output().debugName()
  987. self.name_to_node[output_name] = cond_node
  988. elif node.outputsSize() > 1:
  989. for i, output in enumerate(node.outputs()):
  990. output_name = output.debugName()
  991. getitem = self.fx_graph.call_function(operator.getitem, (cond_node, i))
  992. self.name_to_node[output_name] = getitem
  993. def convert_aten_Bool(self, node: torch._C.Node):
  994. self._convert_as_noop(node)
  995. def convert_prim_Enter(self, node: torch._C.Node):
  996. # export generally treats prim::Enter as noop
  997. # The only context manager export supports is aten::enable_grad.
  998. # Unfortunately, TorchScript does not support aten::enable_grad yet.
  999. # TODO: support aten::enable_grad in both TorchScript and Converter.
  1000. return
  1001. def convert_prim_Exit(self, node: torch._C.Node):
  1002. # export treats prim::Exit as noop
  1003. return
  1004. def _convert_as_noop(self, node: torch._C.Node):
  1005. # Converts the node as a no-op by mapping its output node as arg[0]
  1006. target = get_op_overload(node)
  1007. schema = target._schema
  1008. args, _kwargs = self.get_args_kwargs(node, schema)
  1009. output_name = node.output().debugName()
  1010. self.name_to_node[output_name] = args[0]
  1011. def convert_profiler__record_function_exit(self, node: torch._C.Node):
  1012. # _record_function_exit has side effect so we keep it in fx.graph
  1013. # currently, _record_function_enter_new and _record_function_exit are
  1014. # discarded during `retrace_as_exported_program`.
  1015. target = torch.ops.profiler._record_function_exit
  1016. args = tuple(self.get_fx_value_by_ir_value(input) for input in node.inputs())
  1017. self.fx_graph.call_function(target, args)
  1018. def convert_prim_tolist(self, node: torch._C.Node):
  1019. # prim::tolist cannot be supported by `_convert_standard_operators`
  1020. # since it requires call_method instead of call_function.
  1021. target = "tolist"
  1022. args = (self.get_fx_value_by_ir_value(next(node.inputs())),)
  1023. fx_node = self.fx_graph.call_method(target, args)
  1024. output_name = node.output().debugName()
  1025. self.name_to_node[output_name] = fx_node
  1026. def convert_prim_Uninitialized(self, node: torch._C.Node):
  1027. # `prim::Uninitialized` is inserted by the compiler when it can prove
  1028. # the value will never be used. It can be introduced by exceptions,
  1029. # breaks, continues, and returns.
  1030. # So we add a dummy constant to the graph.
  1031. output_name = node.output().debugName()
  1032. self.name_to_constant[output_name] = torch.Tensor()
  1033. def _convert_standard_operators(self, node: torch._C.Node):
  1034. target = kind_to_standard_operators[node.kind()]
  1035. args = tuple(self.get_fx_value_by_ir_value(input) for input in node.inputs())
  1036. fx_node = self.fx_graph.call_function(target, args)
  1037. output_name = node.output().debugName()
  1038. self.name_to_node[output_name] = fx_node
  1039. def convert_node(self, node: torch._C.Node):
  1040. node_kind = node.kind()
  1041. # Get handler based on namespace and operator name.
  1042. # Provide a default node handler as well in case we don't find
  1043. # matching converter for that.
  1044. handler_func_name = ir_name_to_func_name(node_kind)
  1045. handler_func = getattr(self, handler_func_name, self.convert_call_function_op)
  1046. # str calls print function implemented in CPP. To avoid repeating
  1047. # the entire logic here, we simply keep first line from node string (getting rid
  1048. # of sub-blocks IR prints).
  1049. node_str = "".join(str(node).split("\n")[:1])
  1050. log.debug("[%s] converts [%s]", handler_func.__name__, node_str)
  1051. try:
  1052. handler_func(node)
  1053. except Exception as e:
  1054. raise RuntimeError(f"TS2EPConverter failed for node {node_kind}") from e
  1055. def convert_graph_outputs(self):
  1056. args = []
  1057. outp_name_list = [outp.debugName() for outp in self.ts_graph.outputs()] + list(
  1058. self.name_update_from_subblock_to_parent
  1059. )
  1060. for output_name in outp_name_list:
  1061. if output_name in self.name_to_node:
  1062. fx_node = self.name_to_node[output_name]
  1063. # TODO: Revisit this later after HigherOrderOp design changes.
  1064. # Currently, we cannot directly return input as output.
  1065. if (
  1066. not self.is_top_level_graph()
  1067. and isinstance(fx_node, torch.fx.Node)
  1068. and fx_node.op == "placeholder"
  1069. ):
  1070. fx_node = self.fx_graph.call_function(torch.clone, (fx_node,))
  1071. args.append(fx_node)
  1072. self.output_specs.append(
  1073. OutputSpec(
  1074. OutputKind.USER_OUTPUT,
  1075. arg=TensorArgument(name=output_name),
  1076. target=output_name,
  1077. )
  1078. )
  1079. elif output_name in self.name_to_constant:
  1080. args.append(self.name_to_constant[output_name])
  1081. self.output_specs.append(
  1082. OutputSpec(
  1083. OutputKind.USER_OUTPUT,
  1084. arg=ConstantArgument(
  1085. name=output_name, value=self.name_to_constant[output_name]
  1086. ),
  1087. target=output_name,
  1088. )
  1089. )
  1090. else:
  1091. raise ValueError(f"Output {output_name} not found")
  1092. if len(args) == 0:
  1093. # Sub-block of prim::If can have zero output.
  1094. self.fx_graph.output([])
  1095. elif len(args) == 1:
  1096. self.fx_graph.output(
  1097. args[0]
  1098. ) # Get rid of an extra list wrapped around final output.
  1099. elif len(args) > 1:
  1100. self.fx_graph.output(
  1101. args
  1102. ) # For prim::Loop and prim::If with multiple outputs.
  1103. else:
  1104. # Sub-block of prim::Loop can have multiple outputs.
  1105. self.fx_graph.output(args)
  1106. class ExplainTS2FXGraphConverter(TS2FXGraphConverter):
  1107. """
  1108. Run TS2FXGraphConverter in an explain mode. It collects all failed operators conversions
  1109. and provide that information to users. In order to collect all failed conversions, it
  1110. also mocks some internal attributes (e.g., name_to_node).
  1111. """
  1112. class _DictMock(dict):
  1113. def __init__(self, dict_data, mock_value):
  1114. super().__init__(dict_data)
  1115. self.mock_value = mock_value
  1116. def __getitem__(self, key):
  1117. # If the original dictionary has the key, return its value.
  1118. # Otherwise, return the mock value.
  1119. if not super().__contains__(key):
  1120. return self.mock_value
  1121. return super().__getitem__(key)
  1122. def __contains__(self, key):
  1123. return True
  1124. def __init__(
  1125. self,
  1126. ts_graph: Union[torch._C.Graph, torch._C.Block],
  1127. name_to_param: dict[str, torch.Tensor],
  1128. name_to_buffer: dict[str, torch.Tensor],
  1129. blocks_to_lifted_attrs: dict[torch._C.Block, set[str]],
  1130. name_to_non_tensor_attribute: dict[str, Any],
  1131. name_to_constant: dict[str, Any],
  1132. name_to_attribute_fqn: dict[str, str],
  1133. ):
  1134. super().__init__(
  1135. ts_graph,
  1136. name_to_param,
  1137. name_to_buffer,
  1138. blocks_to_lifted_attrs,
  1139. name_to_non_tensor_attribute,
  1140. name_to_constant,
  1141. name_to_attribute_fqn,
  1142. )
  1143. # Data to keep track of unsupported nodes.
  1144. self.unsupported_node_list: list[torch._C.Node] = []
  1145. # Add mock to needed attributes.
  1146. self.name_to_node = ExplainTS2FXGraphConverter._DictMock(
  1147. self.name_to_node,
  1148. # Dummy node.
  1149. torch.fx.Node(
  1150. None, # type: ignore[arg-type]
  1151. "mock",
  1152. "call_function",
  1153. lambda: None,
  1154. (),
  1155. {},
  1156. ),
  1157. )
  1158. def explain(self):
  1159. self.convert_graph_inputs()
  1160. for node in self.ts_graph.nodes():
  1161. self.convert_node(node)
  1162. self.convert_graph_outputs()
  1163. def convert_node(self, node):
  1164. try:
  1165. super().convert_node(node)
  1166. except Exception:
  1167. self.unsupported_node_list.append(node)
  1168. @contextmanager
  1169. def disable_logging(log):
  1170. disabled = log.disabled
  1171. log.disabled = True
  1172. try:
  1173. yield
  1174. finally:
  1175. log.disabled = disabled
  1176. class TS2EPConverter:
  1177. # TorchScript model to ExportedProgram converter
  1178. def __init__(
  1179. self,
  1180. ts_model: Union[torch.jit.ScriptModule, torch.jit.ScriptFunction],
  1181. sample_args: tuple[Any, ...],
  1182. sample_kwargs: Optional[dict[str, Any]] = None,
  1183. ):
  1184. self.ts_model = ts_model
  1185. self.ts_graph, self.params, _, _ = _create_jit_graph(ts_model, sample_args)
  1186. self.sample_args = sample_args
  1187. self.sample_kwargs = sample_kwargs
  1188. self.name_to_param: dict[str, torch.Tensor] = {}
  1189. self.name_to_buffer: dict[str, torch.Tensor] = {}
  1190. param_list = (
  1191. list(self.ts_model.parameters())
  1192. if not isinstance(self.ts_model, torch._C.ScriptFunction)
  1193. else []
  1194. )
  1195. if not isinstance(self.ts_model, torch._C.ScriptFunction):
  1196. for k, tensor in self.ts_model.state_dict().items(): # type: ignore[union-attr]
  1197. # Check if tensor belongs to any parameter.
  1198. if any(
  1199. (tensor == param).all()
  1200. for param in param_list
  1201. if tensor.shape == param.shape
  1202. ):
  1203. self.name_to_param[k] = tensor
  1204. else:
  1205. self.name_to_buffer[k] = tensor
  1206. self.name_to_non_tensor_attributes: dict[str, Any] = {}
  1207. self.name_to_constant: dict[str, Any] = {}
  1208. self.lift_get_attr()
  1209. def convert(self) -> ExportedProgram:
  1210. log.info(
  1211. """
  1212. TS2EPConverter logging starts from here.
  1213. INFO: (TORCH_LOGS="export" <cmd>)
  1214. * Log TorchScript IR.
  1215. DEBUG: (TORCH_LOGS="+export" <cmd>), additionally
  1216. * Log conversion IR by IR in a format of [<conversion handler name>] converts [<IR>].
  1217. """
  1218. )
  1219. log.info("TorchScript graph\n\n%s\n", self.ts_graph)
  1220. blocks_to_lifted_attrs, name_to_attribute_fqn = get_block_to_lifted_attrs(
  1221. self.ts_graph
  1222. )
  1223. graph_converter = TS2FXGraphConverter(
  1224. self.ts_graph,
  1225. self.name_to_param,
  1226. self.name_to_buffer,
  1227. blocks_to_lifted_attrs,
  1228. self.name_to_non_tensor_attributes,
  1229. self.name_to_constant,
  1230. name_to_attribute_fqn,
  1231. )
  1232. gm = graph_converter.convert()
  1233. # Post-proccessing step to deal with quantized operators.
  1234. replace_quantized_ops_with_standard_ops(gm)
  1235. log.info("GraphModule: %s", gm.print_readable(print_output=False))
  1236. ep = self.retrace_as_exported_program(
  1237. gm,
  1238. graph_converter.name_to_constant,
  1239. )
  1240. log.info("%s", ep)
  1241. # Post-processing step to ensure ExportedProgram has the same state_dict as
  1242. # the original TorchScript model. Throw warnings for additionally populated
  1243. # state_dict entries.
  1244. if not isinstance(self.ts_model, torch._C.ScriptFunction):
  1245. for k, tensor in self.ts_model.state_dict().items(): # type: ignore[union-attr]
  1246. if k not in ep.state_dict:
  1247. warnings.warn(
  1248. f"Manually populate {k} into state_dict ExportedProgram, but it is never used by the ExportedProgram."
  1249. )
  1250. ep.state_dict[k] = tensor
  1251. return ep
  1252. @disable_logging(log)
  1253. def explain(self, print_output=True):
  1254. blocks_to_lifted_attrs, name_to_attribute_fqn = get_block_to_lifted_attrs(
  1255. self.ts_graph
  1256. )
  1257. graph_converter = ExplainTS2FXGraphConverter(
  1258. self.ts_graph,
  1259. self.name_to_param,
  1260. self.name_to_buffer,
  1261. blocks_to_lifted_attrs,
  1262. self.name_to_non_tensor_attributes,
  1263. self.name_to_constant,
  1264. name_to_attribute_fqn,
  1265. )
  1266. graph_converter.explain()
  1267. if len(graph_converter.unsupported_node_list) > 0:
  1268. explain_str = "Unsupported nodes are found in the following list:"
  1269. for i, n in enumerate(graph_converter.unsupported_node_list):
  1270. node_str = "".join(str(n).split("\n")[:1])
  1271. explain_str += f"\n\n {i}. {n.kind()} [{node_str}]"
  1272. else:
  1273. explain_str = "Success!"
  1274. if print_output:
  1275. print(explain_str)
  1276. return explain_str
  1277. def retrace_as_exported_program(
  1278. self,
  1279. gm: torch.fx.GraphModule,
  1280. name_to_constant: dict[str, Any],
  1281. ):
  1282. dynamic_shapes = _tree_map_with_path(
  1283. lambda path, x: (
  1284. [Dim.AUTO] * x.dim() if isinstance(x, torch.Tensor) else None
  1285. ),
  1286. self.sample_args,
  1287. )
  1288. # TODO: adjust input orders to match GraphSignature convention
  1289. ep = torch.export._trace._export(
  1290. gm,
  1291. self.sample_args,
  1292. dynamic_shapes=dynamic_shapes,
  1293. strict=False,
  1294. pre_dispatch=True,
  1295. )
  1296. # Post-processing to make sure the ExportedProgram states are correct.
  1297. # Because during conversion, we set tensor constants as GetAttr,
  1298. # retracing cannot recognize them as tensor constants but instead
  1299. # treat them as buffers. We need to set them again here.
  1300. ep._constants.update(
  1301. {
  1302. k: v
  1303. for k, v in name_to_constant.items()
  1304. if isinstance(v, (torch.Tensor, torch.ScriptObject))
  1305. }
  1306. )
  1307. for k in name_to_constant:
  1308. ep.state_dict.pop(k, None)
  1309. for spec in ep.graph_signature.input_specs:
  1310. # Mark as constant tensors for erroneously traced buffers.
  1311. if spec.kind == InputKind.BUFFER and spec.target in name_to_constant:
  1312. assert isinstance(name_to_constant[spec.target], torch.Tensor), (
  1313. f"{type(name_to_constant[spec.target])} has been erroneously marked as buffer"
  1314. )
  1315. spec.kind = InputKind.CONSTANT_TENSOR
  1316. spec.persistent = None
  1317. ep.verifier().check(ep)
  1318. return ep
  1319. def lift_get_attr(self):
  1320. # This function lifts multiple data types.
  1321. # 1. Tensor constants attributes (e.g., self.data = torch.tensor([2,3]))
  1322. # to buffers. Currently, when there are tensor constants, export
  1323. # would error and ask users to register tensor constants as buffers.
  1324. # Since it is hard to manually do so for TorchScript models
  1325. # (e.g., source code is missing), this function automatically
  1326. # lifts tensor constants to be buffers.
  1327. # 2. ScriptObbject to constant. It will then be converted to getattr in
  1328. # in the fx graph.
  1329. #
  1330. # This function should happen in TS2EPConverter instead of
  1331. # TS2FXGraphConverter since it gets attributes from self.ts_model
  1332. # which is not accessible in TS2FXGraphConverter. It is similar to where
  1333. # we collect self.name_to_param and self.name_to_buffer.
  1334. name_to_attribute_fqn: dict[str, str] = {}
  1335. def get_attr(fqn: str):
  1336. name = fqn.split(".")
  1337. v = self.ts_model
  1338. for n in name:
  1339. v = getattr(v, n)
  1340. return v
  1341. def get_fqn(node: torch._C.Node):
  1342. attr_name = node.s("name")
  1343. input_name = node.input().debugName()
  1344. root_attr_name = name_to_attribute_fqn[input_name]
  1345. attr_fqn = f"{root_attr_name}.{attr_name}" if root_attr_name else attr_name
  1346. return attr_fqn
  1347. def _dfs_get_attr(block):
  1348. for node in block.nodes():
  1349. if node.kind() == "prim::CreateObject":
  1350. output_name = node.output().debugName()
  1351. name_to_attribute_fqn[output_name] = ""
  1352. if node.kind() == "prim::GetAttr":
  1353. attr_fqn = get_fqn(node)
  1354. value = get_attr(attr_fqn)
  1355. output_name = node.output().debugName()
  1356. name_to_attribute_fqn[output_name] = attr_fqn
  1357. if isinstance(value, torch.Tensor):
  1358. if attr_fqn not in self.name_to_buffer:
  1359. # Lift tensor constants to be a buffer
  1360. self.name_to_buffer[attr_fqn] = value
  1361. elif isinstance(value, torch.ScriptObject):
  1362. if attr_fqn not in self.name_to_constant:
  1363. self.name_to_constant[attr_fqn] = value
  1364. else:
  1365. self.name_to_non_tensor_attributes[attr_fqn] = value
  1366. for subblock in node.blocks():
  1367. _dfs_get_attr(subblock)
  1368. _dfs_get_attr(self.ts_graph)