_recursive.py 39 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397398399400401402403404405406407408409410411412413414415416417418419420421422423424425426427428429430431432433434435436437438439440441442443444445446447448449450451452453454455456457458459460461462463464465466467468469470471472473474475476477478479480481482483484485486487488489490491492493494495496497498499500501502503504505506507508509510511512513514515516517518519520521522523524525526527528529530531532533534535536537538539540541542543544545546547548549550551552553554555556557558559560561562563564565566567568569570571572573574575576577578579580581582583584585586587588589590591592593594595596597598599600601602603604605606607608609610611612613614615616617618619620621622623624625626627628629630631632633634635636637638639640641642643644645646647648649650651652653654655656657658659660661662663664665666667668669670671672673674675676677678679680681682683684685686687688689690691692693694695696697698699700701702703704705706707708709710711712713714715716717718719720721722723724725726727728729730731732733734735736737738739740741742743744745746747748749750751752753754755756757758759760761762763764765766767768769770771772773774775776777778779780781782783784785786787788789790791792793794795796797798799800801802803804805806807808809810811812813814815816817818819820821822823824825826827828829830831832833834835836837838839840841842843844845846847848849850851852853854855856857858859860861862863864865866867868869870871872873874875876877878879880881882883884885886887888889890891892893894895896897898899900901902903904905906907908909910
  1. import inspect
  2. import torch
  3. import types
  4. import collections
  5. import textwrap
  6. import functools
  7. import warnings
  8. from typing import Dict, List, Set, Type
  9. import torch._jit_internal as _jit_internal
  10. from torch._sources import fake_range
  11. from torch.jit.frontend import get_default_args, get_jit_class_def, get_jit_def, get_class_properties
  12. from torch.jit._builtins import _find_builtin
  13. from torch.jit._check import AttributeTypeIsSupportedChecker
  14. from torch.jit._state import _python_cu, _add_script_class, _get_script_class
  15. from torch.nn import Module
  16. ScriptMethodStub = collections.namedtuple('ScriptMethodStub', ('resolution_callback', 'def_', 'original_method'))
  17. PropertyStub = collections.namedtuple('PropertyStub', ('resolution_callback', 'def_'))
  18. # TODO: there should be a more principled way of doing this.
  19. ignored_attributes = [
  20. "_version",
  21. "_parameters",
  22. "_buffers",
  23. "_modules",
  24. "_initializing",
  25. "_backward_hooks",
  26. "_forward_hooks",
  27. "_forward_pre_hooks",
  28. "_state_dict_hooks",
  29. "_load_state_dict_pre_hooks",
  30. "dump_patches",
  31. ]
  32. def _compile_and_register_class(obj, rcb, qualified_name):
  33. script_class = _get_script_class(obj)
  34. if not script_class:
  35. ast = get_jit_class_def(obj, obj.__name__)
  36. defaults = torch.jit.frontend.get_default_args_for_class(obj)
  37. script_class = torch._C._jit_script_class_compile(qualified_name, ast, defaults, rcb)
  38. _add_script_class(obj, script_class)
  39. return script_class
  40. def make_stub(func, name):
  41. rcb = _jit_internal.createResolutionCallbackFromClosure(func)
  42. ast = get_jit_def(func, name, self_name="RecursiveScriptModule")
  43. return ScriptMethodStub(rcb, ast, func)
  44. def make_stub_from_method(nn_module, method_name):
  45. func = getattr(nn_module, method_name)
  46. if isinstance(func, ScriptMethodStub):
  47. return func
  48. # Make sure the name present in the resulting AST will match the name
  49. # requested here. The only time they don't match is if you do something
  50. # like:
  51. # def _forward(self):
  52. # pass
  53. # forward = _forward
  54. # In this case, the actual function object will have the name `_forward`,
  55. # even though we requested a stub for `forward`.
  56. return make_stub(func, method_name)
  57. def make_stubs_from_exported_methods(mod):
  58. stubs = []
  59. for name in dir(mod):
  60. item = getattr(mod, name, None)
  61. if (
  62. _jit_internal.get_torchscript_modifier(item)
  63. is _jit_internal.FunctionModifiers.EXPORT
  64. ):
  65. stubs.append(make_stub_from_method(mod, name))
  66. return stubs
  67. def jit_ignored_properties(module):
  68. user_annotated_ignored_attributes = getattr(module, "__jit_ignored_attributes__", list())
  69. def get_properties_names(module):
  70. return set(k for k, v in vars(module).items() if isinstance(v, property))
  71. properties = get_properties_names(type(module))
  72. user_annoted_ignored_properties = set()
  73. for ignored_attr in user_annotated_ignored_attributes:
  74. if ignored_attr in properties:
  75. user_annoted_ignored_properties.add(ignored_attr)
  76. return user_annoted_ignored_properties
  77. # base types that can be constants
  78. # in addition, tuples and lists of these base types are also considered constants
  79. # If you edit this list, then you also need to edit the handlers in
  80. # ConstantValue in jit/script/init.cpp
  81. _constant_types = (bool, float, int, str, type(None), torch.device, torch.layout, torch.dtype)
  82. def _get_valid_constant(attr, v, owner_type):
  83. if isinstance(v, _constant_types):
  84. return v
  85. elif isinstance(v, tuple) or isinstance(v, list):
  86. return tuple(_get_valid_constant(attr, x, owner_type) for x in v)
  87. constants = ", ".join(torch.typename(typ) for typ in _constant_types)
  88. raise TypeError(textwrap.dedent("""
  89. '{}' object in attribute '{}.{}' is not a valid constant.
  90. Valid constants are:
  91. 1. a nn.ModuleList
  92. 2. a value of type {{{}}}
  93. 3. a list or tuple of (2)
  94. """.format(torch.typename(type(v)), owner_type, attr, constants)))
  95. class SourceContext(torch._C._jit_tree_views.SourceRangeFactory):
  96. def __init__(self, source, filename, file_lineno, leading_whitespace_len):
  97. super(SourceContext, self).__init__(source, filename, file_lineno, leading_whitespace_len)
  98. def infer_concrete_type_builder(nn_module, share_types=True):
  99. """
  100. Build a ConcreteModuleTypeBuilder from an nn.Module. This
  101. ConcreteModuleType doesn't have a JIT type associated with it yet, it
  102. must be filled in by the caller.
  103. """
  104. concrete_type_builder = torch._C.ConcreteModuleTypeBuilder(type(nn_module))
  105. if isinstance(nn_module, (torch.nn.ModuleDict)):
  106. concrete_type_builder.set_module_dict()
  107. if isinstance(nn_module, (torch.nn.ModuleList, torch.nn.Sequential)):
  108. concrete_type_builder.set_module_list()
  109. if isinstance(nn_module, (torch.nn.ParameterList)):
  110. concrete_type_builder.set_parameter_list()
  111. if isinstance(nn_module, (torch.nn.ParameterDict)):
  112. concrete_type_builder.set_parameter_dict()
  113. class_annotations = getattr(nn_module, '__annotations__', {})
  114. if isinstance(nn_module, (torch.ao.quantization.QuantWrapper)):
  115. class_annotations = {}
  116. # Get user-annotated ignored attributes.
  117. user_annotated_ignored_attributes = getattr(nn_module, "__jit_ignored_attributes__", list())
  118. concrete_type_builder.add_ignored_attributes(user_annotated_ignored_attributes)
  119. ignored_properties = jit_ignored_properties(nn_module)
  120. # try to infer the type from type annotation or from the object itself
  121. def infer_type(name, item):
  122. # The forward function from Module is special; never use this annotations; we
  123. # need to infer type directly using JIT. I originally wanted to write
  124. # this test as isinstance(class_annotations[name], Callable) but
  125. # isinstance on typing things doesn't seem to work: isinstance(list, Callable)
  126. # is also true!
  127. inferred = False
  128. try:
  129. if name in class_annotations and class_annotations[name] != torch.nn.Module.__annotations__["forward"]:
  130. ann_to_type = torch.jit.annotations.ann_to_type(class_annotations[name], fake_range())
  131. attr_type = torch._C.InferredType(ann_to_type)
  132. elif isinstance(item, torch.jit.Attribute):
  133. ann_to_type = torch.jit.annotations.ann_to_type(item.type, fake_range())
  134. attr_type = torch._C.InferredType(ann_to_type)
  135. else:
  136. attr_type = torch._C._jit_try_infer_type(item)
  137. inferred = True
  138. except RuntimeError as re:
  139. raise RuntimeError(
  140. "Error inferring type for {name}: {item}: {re}".format(name=name, item=item, re=re)
  141. )
  142. return attr_type, inferred
  143. added_names = set()
  144. for name, item in nn_module._parameters.items():
  145. if name in user_annotated_ignored_attributes:
  146. continue
  147. assert item is None or isinstance(item, torch.Tensor)
  148. attr_type, _ = infer_type(name, item)
  149. # We currently have the invariant in various places in our code
  150. # that parameters must be Tensors. However, the nn.Module API also
  151. # allows NoneType parameters. These parameters are not returned as
  152. # part of `parameters()` and its variants, but are available
  153. # through direct attribute access.
  154. concrete_type_builder.add_attribute(name, attr_type.type(), True, False)
  155. added_names.add(name)
  156. for name, item in nn_module._buffers.items():
  157. if name in user_annotated_ignored_attributes:
  158. continue
  159. assert item is None or isinstance(item, torch.Tensor)
  160. attr_type, _ = infer_type(name, item)
  161. concrete_type_builder.add_attribute(name, attr_type.type(), False, True)
  162. added_names.add(name)
  163. for name, item in nn_module._modules.items():
  164. if name in user_annotated_ignored_attributes:
  165. continue
  166. attr_type, _ = infer_type(name, item)
  167. if item is None:
  168. # Modules can be None. We don't have direct support for optional
  169. # Modules, so the register it as an NoneType attribute instead.
  170. concrete_type_builder.add_attribute(name, attr_type.type(), False, False)
  171. continue
  172. if attr_type.success():
  173. assert attr_type.type().is_interface_type()
  174. # if the type can be inferred, it should be a module interface type
  175. sub_concrete_type = torch._C.ConcreteModuleType.from_jit_type(attr_type.type())
  176. else:
  177. # otherwise we get the concrete module type for item and add it to concrete_type
  178. sub_concrete_type = get_module_concrete_type(item, share_types)
  179. concrete_type_builder.add_module(name, sub_concrete_type)
  180. added_names.add(name)
  181. # populate constants_set
  182. constants_set = set(getattr(nn_module, "__constants__", ()))
  183. # Constants annotated via `Final[T]` rather than being added to `__constants__`
  184. for name, ann in class_annotations.items():
  185. if torch._jit_internal.is_final(ann):
  186. constants_set.add(name)
  187. for name in constants_set:
  188. if name in added_names:
  189. # TODO: We should really error in this case, but its bc-breaking so
  190. # we need to warn for at least one release
  191. if name in nn_module._modules:
  192. hint = "submodule"
  193. elif name in nn_module._buffers:
  194. hint = "buffer"
  195. elif name in nn_module._parameters:
  196. hint = "parameter"
  197. else:
  198. raise AssertionError("added_names must be submodule, parameter, or buffer")
  199. warnings.warn("'{}' was found in ScriptModule constants, "
  200. " but it is a non-constant {}. Consider removing it.".format(name, hint))
  201. continue
  202. if not hasattr(nn_module, name):
  203. # TODO: We should really error in this case, but its bc-breaking so
  204. # we need to warn for at least one release
  205. warnings.warn("'{}' was found in ScriptModule constants, "
  206. "but was not actually set in __init__. "
  207. "Consider removing it.".format(name))
  208. continue
  209. value = getattr(nn_module, name)
  210. concrete_type_builder.add_constant(name, _get_valid_constant(name, value, type(nn_module).__name__))
  211. added_names.add(name)
  212. # populate overloads
  213. overloads = getattr(nn_module, "__overloads__", {})
  214. # update with any annotated overloads
  215. overloads.update(get_overload_name_mapping(get_overload_annotations(nn_module, ignored_properties)))
  216. for name, overloaded_names in overloads.items():
  217. concrete_type_builder.add_overload(name, overloaded_names)
  218. for name, value in nn_module.__dict__.items():
  219. if name in ignored_attributes or name.startswith("__"):
  220. # Python objects have lots of random attributes attached to them;
  221. # PyTorch adds a few more. Prevent these from getting compiled.
  222. continue
  223. if name in user_annotated_ignored_attributes:
  224. continue
  225. if name in added_names:
  226. # Don't re-add anything we already added
  227. continue
  228. isoverloadpacket = isinstance(value, torch._ops.OpOverloadPacket)
  229. if isoverloadpacket:
  230. value = value.op
  231. # Handle Python function attributes
  232. if inspect.isfunction(value):
  233. try:
  234. scripted_fn = torch.jit.script(value)
  235. concrete_type_builder.add_function_attribute(
  236. name,
  237. torch._C._jit_try_infer_type(scripted_fn).type(),
  238. value)
  239. except Exception as e:
  240. # If we fail to script the function, it isn't a hard error.
  241. # Instead, we will add it to the list of attributes we failed
  242. # to convert, with the compilation error.
  243. hint = ("(This function exists as an attribute on the Python module, "
  244. "but we failed to compile it to a TorchScript function. "
  245. "\nThe error stack is reproduced here:\n{}").format(e)
  246. concrete_type_builder.add_failed_attribute(name, hint)
  247. pass
  248. continue
  249. # Handle calls to builtin functions (either bespoke builtins from torch.jit._builtins or
  250. # a call to an aten function like torch.add)
  251. builtin_symbol_name = _find_builtin(value)
  252. if builtin_symbol_name:
  253. concrete_type_builder.add_builtin_function(name, builtin_symbol_name)
  254. continue
  255. # Handle Script function attributes
  256. if isinstance(value, torch.jit.ScriptFunction):
  257. concrete_type_builder.add_function_attribute(
  258. name,
  259. torch._C._jit_try_infer_type(value).type(),
  260. value)
  261. continue
  262. # If we got here, this is a regular "data" attribute, add it to the concrete type
  263. attr_type, inferred = infer_type(name, value)
  264. if attr_type.success():
  265. concrete_type_builder.add_attribute(name, attr_type.type(), False, False)
  266. else:
  267. # TODO: could add more detail here. For example, what the user should do
  268. # when the pytype is `list` or `NoneType`
  269. inferred_msg = "Its type was inferred; try adding a type annotation for the attribute." if inferred else ""
  270. additional_info = f"{attr_type.reason()}. {inferred_msg}"
  271. hint = "(This attribute exists on the Python module, " \
  272. f"but we failed to convert Python type: '{torch.typename(type(value))}' " \
  273. f"to a TorchScript type. {additional_info})"
  274. concrete_type_builder.add_failed_attribute(name, hint)
  275. # add hooks to concrete type
  276. for hook in nn_module._forward_hooks.values():
  277. concrete_type_builder.add_forward_hook(hook)
  278. for pre_hook in nn_module._forward_pre_hooks.values():
  279. concrete_type_builder.add_forward_pre_hook(pre_hook)
  280. return concrete_type_builder
  281. class ConcreteTypeStore(object):
  282. type_store: Dict[Type[Module], List[torch._C.ConcreteModuleType]]
  283. methods_compiled: Set[torch._C.ConcreteModuleType]
  284. def __init__(self):
  285. # Python module type => List[ConcreteModuleType)]
  286. self.type_store = {}
  287. # ConcreteTypes that have had their methods already compiled
  288. self.methods_compiled = set()
  289. def get_or_create_concrete_type(self, nn_module):
  290. """
  291. Infer a ConcreteType from this `nn.Module` instance. Underlying JIT
  292. types are re-used if possible.
  293. """
  294. concrete_type_builder = infer_concrete_type_builder(nn_module)
  295. nn_module_type = type(nn_module)
  296. if nn_module_type not in self.type_store:
  297. self.type_store[nn_module_type] = []
  298. # Search the type store for an already-available JIT type
  299. known_types = self.type_store[nn_module_type]
  300. for known_type in known_types:
  301. if known_type.equals(concrete_type_builder):
  302. return known_type
  303. # We didn't find anything; generate a new JIT type from this concrete type
  304. concrete_type = concrete_type_builder.build()
  305. self.type_store[nn_module_type].append(concrete_type)
  306. return concrete_type
  307. concrete_type_store = ConcreteTypeStore()
  308. def create_methods_and_properties_from_stubs(concrete_type, method_stubs, property_stubs):
  309. method_defs = [m.def_ for m in method_stubs]
  310. method_rcbs = [m.resolution_callback for m in method_stubs]
  311. method_defaults = [get_default_args(m.original_method) for m in method_stubs]
  312. property_defs = [p.def_ for p in property_stubs]
  313. property_rcbs = [p.resolution_callback for p in property_stubs]
  314. concrete_type._create_methods_and_properties(property_defs, property_rcbs, method_defs, method_rcbs, method_defaults)
  315. def create_hooks_from_stubs(concrete_type, hook_stubs, pre_hook_stubs):
  316. hook_defs = [h.def_ for h in hook_stubs]
  317. hook_rcbs = [h.resolution_callback for h in hook_stubs]
  318. pre_hook_defs = [h.def_ for h in pre_hook_stubs]
  319. pre_hook_rcbs = [h.resolution_callback for h in pre_hook_stubs]
  320. concrete_type._create_hooks(hook_defs, hook_rcbs, pre_hook_defs, pre_hook_rcbs)
  321. def get_module_concrete_type(nn_module, share_types=True):
  322. """
  323. Gets a concrete type for nn_modules. If share_types is True, the concrete
  324. type is fetched from concrete_type_store. If it is False, a new concrete type
  325. is created without first searching concrete_type_store.
  326. Args:
  327. nn_module: The original Python nn.Module that we are creating a ScriptModule for.
  328. share_types = Whether to share underlying JIT types between modules (if possible).
  329. Returns:
  330. A concrete type for nn_module.
  331. """
  332. assert isinstance(nn_module, Module)
  333. if isinstance(nn_module, torch.jit.ScriptModule) and \
  334. hasattr(nn_module, "_concrete_type"):
  335. return nn_module._concrete_type
  336. if share_types:
  337. # Look into the store of cached JIT types
  338. concrete_type = concrete_type_store.get_or_create_concrete_type(nn_module)
  339. else:
  340. # Get a concrete type directly, without trying to re-use an existing JIT
  341. # type from the type store.
  342. concrete_type_builder = infer_concrete_type_builder(nn_module, share_types)
  343. concrete_type_builder.set_poisoned()
  344. concrete_type = concrete_type_builder.build()
  345. return concrete_type
  346. def create_script_class(obj):
  347. """
  348. Create and return a RecursiveScriptClass instance from a Python object.
  349. Arguments:
  350. obj: A Python object.
  351. """
  352. qualified_class_name = _jit_internal._qualified_name(type(obj))
  353. rcb = _jit_internal.createResolutionCallbackForClassMethods(type(obj))
  354. # Script the type of obj if it hasn't already been scripted.
  355. _compile_and_register_class(type(obj), rcb, qualified_class_name)
  356. class_ty = _python_cu.get_class(qualified_class_name)
  357. # Create an empty torch._C.ScriptObject with the scripted type.
  358. cpp_object = torch._C._create_object_with_type(class_ty)
  359. # Copy all of the attributes over to the torch._C.ScriptObject.
  360. for name, value in obj.__dict__.items():
  361. cpp_object.setattr(name, value)
  362. # Wrap the torch._C.ScriptObject in a RecursiveScriptClass instance.
  363. return wrap_cpp_class(cpp_object)
  364. def create_script_module(nn_module, stubs_fn, share_types=True, is_tracing=False):
  365. """
  366. Creates a new ScriptModule from an nn.Module
  367. Args:
  368. nn_module: The original Python nn.Module that we are creating a ScriptModule for.
  369. stubs_fn: Lambda that takes an nn.Module and generates a list of ScriptMethodStubs to compile.
  370. share_types: Whether to share underlying JIT types between modules (if possible).
  371. NOTE: Only set to False this when we cannot guarantee type sharing will work
  372. correctly. This only happens today for traced modules, where the same
  373. module can produce different traced methods depending on the inputs.
  374. is_tracing: Whether this function is called during tracing or scripting. If tracing,
  375. we don't need to do AttributeTypeIsSupportedChecker because all the unsupported
  376. attributes will be baked as constant in the tracing graph. In addition,
  377. this check significantly slows down the traced modules when the module size is big.
  378. """
  379. assert not isinstance(nn_module, torch.jit.RecursiveScriptModule)
  380. check_module_initialized(nn_module)
  381. concrete_type = get_module_concrete_type(nn_module, share_types)
  382. if not is_tracing:
  383. AttributeTypeIsSupportedChecker().check(nn_module)
  384. return create_script_module_impl(nn_module, concrete_type, stubs_fn)
  385. def create_script_module_impl(nn_module, concrete_type, stubs_fn):
  386. """
  387. Convert an nn.Module to a RecursiveScriptModule.
  388. Args:
  389. nn_module: The original Python nn.Module that we are creating a ScriptModule for.
  390. concrete_type: The fully initialized ConcreteType of the module.
  391. stubs_fn: Lambda that takes an nn.Module and generates a list of ScriptMethodStubs to compile.
  392. """
  393. cpp_module = torch._C._create_module_with_type(concrete_type.jit_type)
  394. method_stubs = stubs_fn(nn_module)
  395. property_stubs = get_property_stubs(nn_module)
  396. hook_stubs, pre_hook_stubs = get_hook_stubs(nn_module)
  397. user_annotated_ignored_attributes = getattr(nn_module, "__jit_ignored_attributes__", list())
  398. ignored_properties = jit_ignored_properties(nn_module)
  399. def init_fn(script_module):
  400. # Initialize the ScriptModule:
  401. # 1. Copy the attributes/parameters/buffers from the original `nn_module` to the new ScriptModule.
  402. for name, (attr_type, is_param) in concrete_type.get_attributes().items():
  403. orig_value = getattr(nn_module, name)
  404. orig_value = orig_value.value if isinstance(orig_value, torch.jit.Attribute) else orig_value
  405. cpp_module.setattr(name, orig_value)
  406. # 2. Copy the submodules from the original `nn_module` to the new ScriptModule,
  407. # recursively scripting them.
  408. for name, sub_concrete_type in concrete_type.get_modules():
  409. orig_value = getattr(nn_module, name)
  410. assert isinstance(orig_value, Module), "Expected Module but got {}".format(type(orig_value))
  411. module_type = sub_concrete_type.jit_type
  412. if isinstance(module_type, torch._C.InterfaceType):
  413. # use the interface inference rule to compile the module
  414. scripted = interface_script(module_type, orig_value)
  415. elif isinstance(orig_value, torch.jit.ScriptModule):
  416. scripted = orig_value
  417. else:
  418. # always reuse the provided stubs_fn to infer the methods to compile
  419. scripted = create_script_module_impl(orig_value, sub_concrete_type, stubs_fn)
  420. cpp_module.setattr(name, scripted)
  421. script_module._modules[name] = scripted
  422. # 3. Copy @ignored/@unused methods and attrs from the original `nn_module` to the new ScriptModule.
  423. # This ensures we can access these Python methods on the ScriptModule.
  424. for name in dir(nn_module):
  425. if name in ignored_properties:
  426. continue
  427. item = getattr(nn_module, name, None)
  428. if inspect.ismethod(item) and _jit_internal.is_ignored_fn(item):
  429. unbound_function = getattr(nn_module, name).__func__
  430. bound_method = unbound_function.__get__(script_module)
  431. setattr(script_module, name, bound_method)
  432. elif concrete_type.is_ignored_attribute(name):
  433. setattr(script_module, name, item)
  434. # For convenience, attach the concrete type to the new ScriptModule
  435. script_module._concrete_type = concrete_type
  436. # Actually create the ScriptModule, initializing it with the function we just defined
  437. script_module = torch.jit.RecursiveScriptModule._construct(cpp_module, init_fn)
  438. # Compile methods if necessary
  439. if concrete_type not in concrete_type_store.methods_compiled:
  440. create_methods_and_properties_from_stubs(concrete_type, method_stubs, property_stubs)
  441. # Create hooks after methods to ensure no name collisions between hooks and methods.
  442. # If done before, hooks can overshadow methods that aren't exported.
  443. create_hooks_from_stubs(concrete_type, hook_stubs, pre_hook_stubs)
  444. torch._C._run_emit_module_hook(cpp_module)
  445. concrete_type_store.methods_compiled.add(concrete_type)
  446. # Copy the forward hooks and pre-hooks to the new ScriptModule
  447. # to allow the hooks to be run from eager as ScriptFunctions
  448. for idx, fn in enumerate(script_module._c._get_forward_pre_hooks()):
  449. script_module._forward_pre_hooks[idx] = fn
  450. for idx, fn in enumerate(script_module._c._get_forward_hooks()):
  451. script_module._forward_hooks[idx] = fn
  452. # Special handling so methods like __len__ work in script methods on classes derived from containers
  453. if isinstance(nn_module, (torch.nn.ModuleList, torch.nn.Sequential, torch.nn.ModuleDict)) and \
  454. '__len__' not in cpp_module._method_names():
  455. script_module.define("def __len__(self):\n return {}\n".format(len(nn_module)))
  456. if isinstance(nn_module, torch.nn.ModuleDict) and \
  457. '__contains__' not in cpp_module._method_names():
  458. if len(nn_module.keys()):
  459. keys = repr(list(nn_module.keys()))
  460. script_module.define("def __contains__(self, key: str):\n return key in {}\n".format(keys))
  461. else:
  462. script_module.define("def __contains__(self, key: str):\n return False\n")
  463. # Make the compiled methods available to the Python ScriptModule class.
  464. for method_stub in method_stubs:
  465. if method_stub.original_method is None:
  466. # define()'d methods don't have an Python original_method, so we
  467. # don't need to do any Python re-wrapping stuff
  468. continue
  469. name = method_stub.original_method.__name__
  470. if name != method_stub.def_.name().name:
  471. # TODO: Why skip this? Because @torch.jit._overload_method will
  472. # mangle the name of the function.
  473. continue
  474. script_method = cpp_module._get_method(name)
  475. # Wrap the original to propagate docstrings and such.
  476. # TODO: we don't currently do this functions that are recursively
  477. # compiled, we should.
  478. wrapped_script_method = functools.wraps(method_stub.original_method)(script_method)
  479. # Add the methods to the script_module directly. This ensures they will
  480. # be found first when `name` is looked up (as opposed to the stubs or
  481. # nn.Module.forward)
  482. script_module.__dict__[name] = wrapped_script_method
  483. # Make module properties available on the Python ScriptModule class.
  484. for property_stub in property_stubs:
  485. property_name = property_stub.def_.name().name
  486. fget = cpp_module._get_method(property_stub.def_.getter_name().name)
  487. # Setter is optional, so it may not exist.
  488. setter_name = property_stub.def_.setter_name()
  489. fset = cpp_module._get_method(setter_name.name) if setter_name else None
  490. script_module.__dict__[property_name] = property(property_name, fget, fset) # type: ignore[arg-type]
  491. # copy over python methods to script module if they aren't defined on the script module
  492. # this is currently an internal api used only on module containers
  493. for name in dir(nn_module):
  494. if name in ignored_properties:
  495. continue
  496. item = getattr(nn_module, name, None)
  497. if _jit_internal.get_torchscript_modifier(item) is _jit_internal.FunctionModifiers.COPY_TO_SCRIPT_WRAPPER:
  498. add_python_attr_to_scripted_model(script_module, nn_module, name)
  499. return script_module
  500. # We define shims of certain attributes on the RecursiveScriptModule to support
  501. # magic methods. To check if a script model defines an attribute we need
  502. # to also check that the attribute is not the shim
  503. def script_model_defines_attr(script_model, attr):
  504. script_attr = getattr(script_model, attr, None)
  505. if script_attr is None:
  506. return False
  507. default_attr = getattr(torch.jit.RecursiveScriptModule, attr, None)
  508. if default_attr is None:
  509. return False
  510. return script_attr != default_attr
  511. def add_python_attr_to_scripted_model(script_model, orig, attr):
  512. if hasattr(orig, attr) and script_model_defines_attr(script_model, attr):
  513. setattr(script_model, attr, getattr(orig, attr))
  514. def get_overload_annotations(mod, jit_ignored_properties):
  515. # original function => [(mangled overload name, overload function)]
  516. overloads = {}
  517. for name in dir(type(mod)):
  518. if name in jit_ignored_properties:
  519. continue
  520. item = getattr(mod, name, None)
  521. if not callable(item):
  522. continue
  523. # builtin functions like repr() in python 2 do not have __module__ defined
  524. if hasattr(item, "__module__") and item.__module__ is not None:
  525. method_overloads = _jit_internal._get_overloaded_methods(item, mod.__class__)
  526. if method_overloads is None:
  527. continue
  528. if item.__func__ in method_overloads:
  529. raise RuntimeError(_jit_internal.get_overload_no_implementation_error_message(
  530. 'method', item.__func__))
  531. names = [name + "__" + str(i) for i in range(len(method_overloads))]
  532. overloads[item] = list(zip(names, method_overloads))
  533. return overloads
  534. def get_overload_name_mapping(overload_info):
  535. # Same format as __overloads__
  536. # original function => [overload names]
  537. overload_name_mappings: Dict[str, List[str]] = {}
  538. for orig_fn, overloads in overload_info.items():
  539. original_name = orig_fn.__name__
  540. if original_name not in overload_name_mappings:
  541. overload_name_mappings[original_name] = []
  542. for overload_name, _ in overloads:
  543. overload_name_mappings[original_name].append(overload_name)
  544. return overload_name_mappings
  545. def _check_no_signature(func):
  546. signature = torch.jit.annotations.get_signature(func, None, fake_range(), inspect.ismethod(func))
  547. if signature is None:
  548. qual_name = _jit_internal._qualified_name(func)
  549. raise RuntimeError("Must explicitly add type annotations to overloaded functions: {}".format(qual_name))
  550. def make_stubs_for_overloads(overload_info):
  551. overload_stubs = []
  552. for orig_fn, overloads in overload_info.items():
  553. orig_ast = get_jit_def(orig_fn, orig_fn.__name__, self_name="RecursiveScriptModule")
  554. for overload_name, overload_fn in overloads:
  555. _check_no_signature(overload_fn)
  556. over_ast = get_jit_def(overload_fn, overload_fn.__name__, self_name="RecursiveScriptModule")
  557. new_ast = torch._C._replace_overloaded_method_decl(over_ast.decl(), orig_ast, overload_name)
  558. _rcb = _jit_internal.createResolutionCallbackFromClosure(orig_fn)
  559. overload_stubs.append(ScriptMethodStub(_rcb, new_ast, overload_fn))
  560. return overload_stubs
  561. def check_module_initialized(mod):
  562. assert isinstance(mod, torch.nn.Module)
  563. if not hasattr(mod, '_parameters'):
  564. raise RuntimeError("'{}' has not been initialized, did you forget to call 'super()'?"
  565. .format(torch.typename(type(mod))))
  566. # This is to avoid importing torch.distributed.nn
  567. if not hasattr(mod, 'remote_parameters'):
  568. for name, param in mod._parameters.items():
  569. if param is not None and torch.nn.parameter.is_lazy(param):
  570. raise RuntimeError("'{}' has uninitialized parameters {}. Did you forget to run a forward pass?"
  571. .format(torch.typename(type(mod)), name))
  572. for name, buf in mod._buffers.items():
  573. if buf is not None and torch.nn.parameter.is_lazy(buf):
  574. raise RuntimeError("'{}' has uninitialized buffers {}. Did you forget to run a forward pass?"
  575. .format(torch.typename(type(mod)), name))
  576. def infer_methods_to_compile(nn_module):
  577. """
  578. Implements the default rules for which methods should act as starting
  579. points for compilation (TODO add a link when the rules are published).
  580. """
  581. check_module_initialized(nn_module)
  582. user_annotated_ignored_attributes = getattr(nn_module, "__jit_ignored_attributes__", list())
  583. ignored_properties = jit_ignored_properties(nn_module)
  584. methods: List[str] = []
  585. if hasattr(nn_module, 'forward') and not _jit_internal.is_ignored_fn(nn_module.forward):
  586. forward_func = getattr(nn_module.forward, "__func__", None)
  587. module_forward = getattr(torch.nn.Module, "forward", None)
  588. if forward_func != module_forward:
  589. methods = ['forward']
  590. exported = []
  591. for name in dir(nn_module):
  592. if name in ignored_properties:
  593. continue
  594. item = getattr(nn_module, name, None)
  595. if _jit_internal.get_torchscript_modifier(item) is _jit_internal.FunctionModifiers.EXPORT:
  596. exported.append(name)
  597. methods = methods + exported
  598. overload_name_mappings = dict(getattr(nn_module, "__overloads__", {}))
  599. overload_info = get_overload_annotations(nn_module, ignored_properties)
  600. overload_name_mappings.update(get_overload_name_mapping(overload_info))
  601. overload_stubs = make_stubs_for_overloads(overload_info)
  602. nn_module.__overloads__ = overload_name_mappings
  603. # we shouldn't directly compile overloaded methods, just its overloads
  604. def ignore_overloaded(method_name):
  605. return method_name not in overload_name_mappings
  606. filtered_methods = filter(ignore_overloaded, methods)
  607. # Unique the methods. We don't want to use a set to store the methods because it
  608. # introduces non-determinism to compile order.
  609. uniquer: Set[str] = set()
  610. uniqued_methods = []
  611. for name in filtered_methods:
  612. if name in uniquer:
  613. continue
  614. uniqued_methods.append(name)
  615. uniquer.add(name)
  616. stubs = []
  617. for method in uniqued_methods:
  618. stubs.append(make_stub_from_method(nn_module, method))
  619. return overload_stubs + stubs
  620. def get_hook_stubs(nn_module):
  621. """
  622. Returns forward hook and pre_hook ScriptModuleStubs
  623. """
  624. check_module_initialized(nn_module)
  625. hook_map: Dict = {}
  626. hook_stubs = []
  627. for hook in nn_module._forward_hooks.values():
  628. if hook.__name__ in hook_map:
  629. if id(hook) != id(hook_map[hook.__name__]):
  630. raise RuntimeError(
  631. f"Hook '{hook.__name__}' on {type(nn_module).__name__} "
  632. "has at least two different python definitions."
  633. " Please use unique names for all hooks."
  634. )
  635. else:
  636. hook_map[hook.__name__] = hook
  637. hook_stubs.append(make_stub(hook, hook.__name__))
  638. pre_hook_stubs = []
  639. for pre_hook in nn_module._forward_pre_hooks.values():
  640. if pre_hook.__name__ in hook_map:
  641. if id(pre_hook) != id(hook_map[pre_hook.__name__]):
  642. raise RuntimeError(
  643. f"Pre-hook '{pre_hook.__name__}' on {type(nn_module).__name__} "
  644. "has at least two different python definitions."
  645. " Please use unique names for all hooks."
  646. )
  647. else:
  648. hook_map[pre_hook.__name__] = pre_hook
  649. pre_hook_stubs.append(make_stub(pre_hook, pre_hook.__name__))
  650. return hook_stubs, pre_hook_stubs
  651. def get_property_stubs(nn_module):
  652. """
  653. Create property stubs for the properties of the module by creating method
  654. stubs for the getter and setter.
  655. """
  656. module_ty = type(nn_module)
  657. properties_asts = get_class_properties(module_ty, self_name="RecursiveScriptModule")
  658. rcbs = {}
  659. for name in dir(module_ty):
  660. item = getattr(module_ty, name, None)
  661. if isinstance(item, property):
  662. if not item.fget:
  663. raise RuntimeError(f'Property {name} of {nn_module.__name__} must have a getter')
  664. rcbs[name] = _jit_internal.createResolutionCallbackFromClosure(item.fget)
  665. stubs = [PropertyStub(rcbs[ast.name().name], ast) for ast in properties_asts]
  666. return stubs
  667. def interface_script(mod_interface, nn_module):
  668. """
  669. Makes a ScriptModule from an nn.Module, using the interface methods rule for
  670. determining which methods to compile.
  671. Args:
  672. mod_interface: the interface type that the module have
  673. nn_module: The original Python nn.Module that we are creating a ScriptModule for.
  674. """
  675. if isinstance(nn_module, torch.jit.ScriptModule):
  676. return nn_module
  677. check_module_initialized(nn_module)
  678. def infer_interface_methods_to_compile(nn_module):
  679. """
  680. Rule to infer the methods from the interface type to know which
  681. methods need to act as starting points for compilation.
  682. """
  683. stubs = []
  684. for method in mod_interface.getMethodNames():
  685. stubs.append(make_stub_from_method(nn_module, method))
  686. return stubs
  687. return create_script_module(nn_module, infer_interface_methods_to_compile)
  688. def try_compile_fn(fn, loc):
  689. if _jit_internal.is_ignored_fn(fn):
  690. # Don't do anything for @ignore'd functions
  691. return None
  692. if isinstance(fn, torch.nn.Module):
  693. # Since modules are callable pybind recognizes them as functions, but
  694. # don't do anything for them
  695. return None
  696. if not inspect.isfunction(fn) and not inspect.ismethod(fn):
  697. raise RuntimeError("`{}` is not a function. Recursive scripting only supports "
  698. "Python functions or methods currently.\n"
  699. "Consider manually annotating `{}` with @torch.jit.script.".format(fn, fn))
  700. # We don't have the actual scope where the function was defined, but we can
  701. # extract the necessary info from the closed over variables on the function
  702. # object
  703. rcb = _jit_internal.createResolutionCallbackFromClosure(fn)
  704. return torch.jit.script(fn, _rcb=rcb)
  705. def wrap_cpp_class(cpp_class):
  706. """
  707. Wrap this torch._C.Object in a Python RecursiveScriptClass.
  708. """
  709. return torch.jit.RecursiveScriptClass(cpp_class)
  710. def wrap_cpp_module(cpp_module):
  711. """
  712. Wrap this torch._C.ScriptModule in a Python ScriptModule, recursively for all submodules
  713. """
  714. def init_fn(script_module):
  715. for name, cpp_module in torch._C.ModuleDict(script_module._c).items():
  716. setattr(script_module, name, wrap_cpp_module(cpp_module))
  717. script_module._concrete_type = torch._C.ConcreteModuleType.from_jit_type(script_module._c._type())
  718. for idx, fn in enumerate(script_module._c._get_forward_pre_hooks()):
  719. script_module._forward_pre_hooks[idx] = fn
  720. for idx, fn in enumerate(script_module._c._get_forward_hooks()):
  721. script_module._forward_hooks[idx] = fn
  722. return torch.jit.RecursiveScriptModule._construct(cpp_module, init_fn)
  723. def compile_unbound_method(concrete_type, fn):
  724. if _jit_internal.is_ignored_fn(fn):
  725. return None
  726. stub = make_stub(fn, fn.__name__)
  727. with torch._jit_internal._disable_emit_hooks():
  728. # We don't want to call the hooks here since the graph that is calling
  729. # this function is not yet complete
  730. create_methods_and_properties_from_stubs(concrete_type, (stub,), ())
  731. return stub
  732. def lazy_bind(concrete_type, unbound_method):
  733. """
  734. Returns a function that lazily binds `unbound_method` to a provided
  735. Module IValue, then invokes the method. We do this so that any Python
  736. shenanigans that will poison type sharing are impossible at compile
  737. time.
  738. """
  739. def lazy_binding_method(cpp_module, *args):
  740. def init_fn(script_module):
  741. orig_class = concrete_type.py_class
  742. # Copy @ignored/@unused methods from the original module to the new one.
  743. # This ensures they are available during execution.
  744. for name in dir(orig_class):
  745. item = getattr(orig_class, name, None)
  746. if _jit_internal.is_ignored_fn(item):
  747. setattr(script_module, name, item)
  748. # Copy constants over so they are available during execution.
  749. for name, value in concrete_type.get_constants().items():
  750. setattr(script_module, name, value)
  751. script_module = torch.jit.RecursiveScriptModule._construct(cpp_module, init_fn)
  752. method = types.MethodType(unbound_method, script_module)
  753. return method(*args)
  754. # make the lazy binding method "look like" the original method
  755. lazy_binding_method.original_fn = unbound_method # type: ignore[attr-defined]
  756. lazy_binding_method.__name__ = unbound_method.__name__
  757. torch._jit_internal.copy_torchscript_modifier(unbound_method, lazy_binding_method)
  758. return lazy_binding_method