You can not select more than 25 topics Topics must start with a chinese character,a letter or number, can include dashes ('-') and can be up to 35 characters long.

network.py 28 kB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397398399400401402403404405406407408409410411412413414415416417418419420421422423424425426427428429430431432433434435436437438439440441442443444445446447448449450451452453454455456457458459460461462463464465466467468469470471472473474475476477478479480481482483484485486487488489490491492493494495496497498499500501502503504505506507508509510511512513514515516517518519520521522523524525526527528529530531532533534535536537538539540541542543544545546547548549550551552553554555556557558559560561562563564565566567568569570571572573574575576577578579580581582583584585586587588589590591592593594595596597598599600601602603604605606607608609610611612613614615616617618619620621622623624625626627628629630631632633634635636637638639640641642643644645646647648649650651652653654655656657658659660661662663664665666667668669670671672673674675676677678679680681682683684685686687688689690691692693694695696697698699700701702703704705706707708709710711712713714715716717718719720721722723724725726727728729730731732733734735736737738739740741742743744745746747748749750751752753754755756757758759760761762763764765766767768769770771772773774775776777778779780781782783784785786787788789790791792793794795796797798799800801802803804805806807808809810811812813814815816817818
  1. # -*- coding: utf-8 -*-
  2. import collections
  3. import fnmatch
  4. import itertools
  5. import pickle
  6. import re
  7. from collections import OrderedDict
  8. from typing import Any, Dict, List, Optional, Sequence
  9. from ..core import _imperative_rt
  10. from ..core._imperative_rt import ComputingGraph, SerializationMetadata
  11. from ..core._trace_option import set_symbolic_shape as _set_symbolic_shape
  12. from ..core.tensor import megbrain_graph as G
  13. from ..logger import get_logger
  14. from .comp_graph_tools import get_dep_vars, get_opr_type, get_oprs_seq
  15. from .network_node import (
  16. ConstOpBase,
  17. Host2DeviceCopy,
  18. ImmutableTensor,
  19. NetworkNode,
  20. OpNode,
  21. VarNode,
  22. str_to_mge_class,
  23. )
  24. logger = get_logger(__name__)
  25. class Network:
  26. def __init__(self):
  27. self.input_vars = [] # input var of graph
  28. self._orig_inputs = []
  29. self.output_vars = [] # output var of graph
  30. self._orig_outputs = []
  31. self.all_oprs_map = OrderedDict() # _imperative_rt.graph.VarNode.id: VarNode
  32. self.all_vars_map = (
  33. OrderedDict()
  34. ) # _imperative_rt.graph.OperatorNode.id: OpNode
  35. self.graph = ComputingGraph()
  36. self._metadata = None
  37. @property
  38. def metadata(self):
  39. r"""Load metadata as a dict."""
  40. if not self._metadata.is_valid:
  41. logger.info("metadata is not valid!")
  42. return None
  43. ret = dict()
  44. try:
  45. user_info = pickle.loads(self._metadata.user_info)
  46. except: # pylint: disable=bare-except
  47. logger.warning(
  48. "can't parse user info by pickle, so return the original bytes object!"
  49. )
  50. user_info = self._metadata.user_info
  51. ret["user_info"] = user_info
  52. ret["graph_modified"] = self._metadata.graph_modified
  53. ret["optimized_for_inference"] = self._metadata.optimized_for_inference
  54. if ret["optimized_for_inference"]:
  55. ret.update(G.deserialize_infer_option(self._metadata.optimize_options))
  56. return ret
  57. @classmethod
  58. def load(cls, model_path: str, outspec: List[str] = None):
  59. r"""Loads a computing graph as a Network object.
  60. Args:
  61. model_path: file path of mge model.
  62. outspec: only load the subgraph with outspec as its endpoints.
  63. """
  64. self = cls()
  65. ret = G.load_graph(model_path)
  66. outputs, self._metadata = ret.output_vars_list, ret.metadata
  67. if outspec is not None:
  68. output_spec = outspec.copy()
  69. all_vars = get_dep_vars(outputs) + outputs
  70. new_outputs = {}
  71. for i in all_vars:
  72. if i.name in output_spec:
  73. new_outputs[i.name] = i
  74. output_spec.remove(i.name)
  75. assert len(output_spec) == 0, "Can not find {} in this model".format(
  76. output_spec
  77. )
  78. outputs = [new_outputs[i] for i in outspec]
  79. self._orig_outputs = outputs
  80. for x in self._orig_outputs:
  81. self.output_vars.append(self._get_var(x))
  82. self.add_dep_oprs()
  83. for x in self._orig_inputs:
  84. self.input_vars.append(self._get_var(x))
  85. self.graph = self._orig_outputs[0].graph
  86. return self
  87. def _compile(self):
  88. self.all_oprs_map = {}
  89. self.all_vars_map = {}
  90. for opr in self.all_oprs:
  91. if isinstance(opr, (ConstOpBase, Host2DeviceCopy)):
  92. opr.compile(self.graph)
  93. else:
  94. opr.compile()
  95. if opr.name is not None:
  96. opr._opr.name = opr.name
  97. self.all_oprs_map[opr._opr.id] = opr
  98. for o in opr.outputs:
  99. self.all_vars_map[o.var.id] = o
  100. def optimize_for_inference(self, dest_vars, **kwargs):
  101. r"""Applies optimize_for_inference pass for operator graph.
  102. Args:
  103. dest_vars: list of output vars in the operator graph
  104. Keyword Arguments:
  105. * enable_io16xc32 --
  106. whether to use float16 for I/O between oprs and use
  107. float32 as internal computation precision. Note the output var would be
  108. changed to float16.
  109. * enable_ioc16 --
  110. whether to use float16 for both I/O and computation
  111. precision.
  112. * enable_hwcd4 --
  113. whether to use NHWCD4 data layout. This is faster on some
  114. OpenCL backend.
  115. * enable_nchw88 --
  116. whether to use NCHW88 data layout, currently
  117. used in X86 AVX backend.
  118. * enable_nchw44 --
  119. whether to use NCHW44 data layout, currently
  120. used in arm backend.
  121. * enable_nchw44_dot --
  122. whether to use NCHW44_dot data layout, currently
  123. used in armv8.2+dotprod backend.
  124. * enable_nchw4 --
  125. whether to use NCHW4 data layout, currently
  126. used in nvidia backend(based on cudnn).
  127. * enable_nchw32 --
  128. whether to use NCHW32 data layout, currently
  129. used in nvidia backend with tensorcore(based on cudnn).
  130. * enable_chwn4 --
  131. whether to use CHWN4 data layout, currently
  132. used in nvidia backend with tensorcore.
  133. * enable_nchw64 --
  134. whether to use NCHW64 data layout, used for fast int4
  135. support on Nvidia GPU.
  136. * enable_fuse_conv_bias_nonlinearity: whether to fuse conv+bias+nonlinearty
  137. into one opr.
  138. * enable_fuse_conv_bias_with_z: whether to fuse conv_bias with z
  139. input for inference on nvidia backend(this optimization pass will
  140. result in mismatch of the precision of output of training and
  141. inference
  142. * enable_fuse_grain: fuse grain will be enable by default to fuse grain operator to huge operator, you can disable it.
  143. )
  144. """
  145. if not isinstance(dest_vars, Sequence):
  146. dest_vars = [dest_vars]
  147. dest_vars = list(G.VarNode(var.var) for var in dest_vars)
  148. new_vars = G.optimize_for_inference(dest_vars, **kwargs)
  149. return list(self._get_var(var) for var in new_vars)
  150. def dump(
  151. self,
  152. file,
  153. *,
  154. keep_var_name: int = 1,
  155. keep_opr_name: bool = False,
  156. keep_param_name: bool = False,
  157. keep_opr_priority: bool = False,
  158. strip_info_file=None,
  159. append_json=False,
  160. optimize_for_inference=True,
  161. append=False,
  162. user_info: Any = None,
  163. enable_metadata=True,
  164. **kwargs
  165. ):
  166. r"""Serializes graph to file.
  167. Args:
  168. file: output file, could be file object or filename.
  169. append: whether output is appended to ``file``.
  170. Only works when ``file`` is str.
  171. keep_var_name: level for keeping variable names:
  172. * 0: none of the names are kept
  173. * 1: (default)keep names of output vars
  174. * 2: keep names of all (output and internal) vars
  175. keep_opr_name: whether to keep operator names.
  176. keep_param_name: whether to keep param names, so param values can be
  177. easily manipulated after loading model
  178. keep_opr_priority: whether to keep priority setting for operators
  179. strip_info_file: a string for path or a file handler. if is not None,
  180. then the dump information for code strip would be written to ``strip_info_file``
  181. append_json: will be check when `strip_info_file` is not None. if set
  182. true, the information for code strip will be append to strip_info_file.
  183. if set false, will rewrite strip_info_file
  184. optimize_for_inference: enbale optmizations,
  185. will skip all optimize options if this is False. Default: True
  186. user_info: any type object, which will be pickled to bytes.
  187. enable_metadata: whether to save metadata into output file.
  188. See more detials in :meth:`~.trace.dump`.
  189. """
  190. def _set_var_name(var):
  191. graph_var = G.VarNode(var.var)
  192. graph_var.name = var.name
  193. return graph_var
  194. self._compile()
  195. out = list(map(_set_var_name, self.output_vars))
  196. if kwargs.pop("arg_names", False):
  197. logger.warning(
  198. '"arg_names" is not supported in Network.dump, rename input vars directly'
  199. )
  200. if kwargs.pop("output_names", False):
  201. logger.warning(
  202. '"output_names" is not supported in Network.dump, rename output vars directly'
  203. )
  204. if optimize_for_inference:
  205. out, optimize_options = G.optimize_for_inference(out, **kwargs)
  206. metadata = SerializationMetadata()
  207. if enable_metadata:
  208. metadata.is_valid = True
  209. metadata.graph_modified = True
  210. metadata.user_info = pickle.dumps(user_info)
  211. if optimize_for_inference:
  212. metadata.optimize_options = optimize_options
  213. G.set_priority_to_id([o._node if isinstance(o, G.VarNode) else o for o in out])
  214. dump_content, dump_info = G.dump_graph(
  215. out,
  216. keep_var_name=keep_var_name,
  217. keep_opr_name=keep_opr_name,
  218. keep_param_name=keep_param_name,
  219. keep_opr_priority=keep_opr_priority,
  220. strip_info_file=strip_info_file,
  221. append_json=append_json,
  222. metadata=metadata,
  223. )
  224. if isinstance(file, str):
  225. permission = "wb" if append == False else "ab"
  226. file = open(file, permission)
  227. file.write(dump_content)
  228. return dump_info
  229. def make_const(self, data, name=None, device=None):
  230. r"""Makes an ImmutableTensor OpNode to provide a parameter for the network."""
  231. node = ImmutableTensor(data, name, device, self.graph)
  232. node.compile(self.graph)
  233. return node.outputs[0]
  234. def make_input_node(self, shape, dtype, name=None, device=None):
  235. r"""Makes a Host2DeviceCopy OpNode to provide an input varnode for the network."""
  236. node = Host2DeviceCopy(shape, dtype, name, device)
  237. node.compile(self.graph)
  238. return node.outputs[0]
  239. def add_output(self, *vars: VarNode):
  240. r"""Adds vars into the network output node list"""
  241. if not all([var.owner for var in vars]):
  242. self.add_dep_oprs(*vars)
  243. for var in vars:
  244. # use method 'is' instead of 'in' to avoid
  245. # compare VarNode use elemwise equal
  246. if not any(var is _ for _ in self.output_vars):
  247. self.output_vars.append(var)
  248. def remove_output(self, *vars: VarNode):
  249. r"""Removes vars from the network output node list"""
  250. for var in vars:
  251. # use list pop instead of remove to avoid
  252. # compare VarNode use elemwise equal
  253. is_removed = False
  254. for idx, out_var in enumerate(self.output_vars):
  255. if var is out_var:
  256. self.output_vars.pop(idx)
  257. is_removed = True
  258. if not is_removed:
  259. logger.warning(
  260. "Failed to remove {}({}). Please check whether "
  261. "this node is in the output list.".format(var.name, id(var))
  262. )
  263. def add_dep_oprs(self, *vars):
  264. if len(vars) == 0:
  265. vars = self.output_vars
  266. assert all(isinstance(var, VarNode) for var in vars), "Only support add VarNode"
  267. q = list(vars)
  268. while len(q) > 0:
  269. cur = q.pop(0)
  270. if cur.owner is not None:
  271. continue
  272. if cur.name is None:
  273. cur.name = cur.var.name
  274. self.all_vars_map[cur.var.id] = cur
  275. mge_opr = cur.var.owner
  276. if get_opr_type(mge_opr) == "Host2DeviceCopy":
  277. self._orig_inputs.extend(mge_opr.outputs)
  278. cur.owner = self._add_opr(mge_opr)
  279. if cur.owner is None:
  280. cur.owner = self.all_oprs_map[mge_opr.id]
  281. continue
  282. q.extend(cur.owner.inputs)
  283. return list(vars)
  284. def modify_opr_names(self, modifier):
  285. r"""Modifies names of operators **inplace**; useful for merging loaded
  286. network into another network
  287. Args:
  288. modifier(str or callable): a string to be prepended to the name, or a function
  289. that maps from name to name
  290. """
  291. if isinstance(modifier, str):
  292. om = modifier
  293. modifier = lambda v: "{}.{}".format(om, v)
  294. assert isinstance(modifier, collections.abc.Callable)
  295. for i in self.all_oprs:
  296. v0 = i.name
  297. v1 = modifier(v0)
  298. assert isinstance(v1, str)
  299. i.name = v1
  300. def reset_batch_size(self, batchsize, *, blacklist=()):
  301. r"""Helper for reset batch size; first dimension of all data providers
  302. not in blacklist are assumed to be the batch size
  303. Args:
  304. blacklist: data provider names whose first dimension is not
  305. batchbatch size
  306. """
  307. blacklist = set(blacklist)
  308. prev_batchsize = None
  309. for i in self.data_providers_filter:
  310. if i.name in blacklist:
  311. blacklist.remove(i.name)
  312. else:
  313. shp = list(i.shape)
  314. if prev_batchsize is None:
  315. prev_batchsize = shp[0]
  316. else:
  317. assert prev_batchsize == shp[0], (
  318. "batchsize mismatch: batchsize={} "
  319. "shape={} dp={}".format(prev_batchsize, shp, i.name)
  320. )
  321. shp[0] = batchsize
  322. i.shape = tuple(shp)
  323. self._compile()
  324. assert prev_batchsize is not None, "no data provider found"
  325. assert not blacklist, "unused items in blacklist: {}".format(blacklist)
  326. def replace_vars(self, repl_dict: Dict[VarNode, VarNode]):
  327. r"""Replaces vars in the graph.
  328. Args:
  329. repl_dict: the map {old_var: new_var} that specifies how to replace the vars.
  330. """
  331. if not all([var.owner for var in repl_dict.values()]):
  332. self.add_dep_oprs(*list(repl_dict.values()))
  333. for var in self.all_vars:
  334. if var in repl_dict:
  335. repl_var = repl_dict[var]
  336. if repl_var is var:
  337. continue
  338. for opnode in var.users:
  339. # use method 'is' instead of 'in' to avoid
  340. # compare VarNode use elemwise equal
  341. assert any([var is _ for _ in opnode.inputs])
  342. opnode.inputs = [repl_var if var is i else i for i in opnode.inputs]
  343. if opnode not in repl_var.users:
  344. repl_var.users.append(opnode)
  345. var.users.clear()
  346. self._compile()
  347. def replace_oprs(self, repl_dict: Dict[OpNode, OpNode]):
  348. r"""Replaces operators in the graph.
  349. Args:
  350. repl_dict: the map {old_opr: new_opr} that specifies how to replace the operators.
  351. """
  352. for opr in self.all_oprs:
  353. if opr in repl_dict:
  354. assert len(opr.outputs) == len(
  355. repl_dict[opr].outputs
  356. ), "can not replace {} with {}".format(type(opr), type(repl_dict[opr]))
  357. for ind, var in enumerate(opr.outputs):
  358. var.owner = repl_dict[opr]
  359. var.__dict__.update(repl_dict[opr].outputs[ind].__dict__)
  360. var._reset_var(repl_dict[opr].outputs[ind].var)
  361. repl_dict[opr].outputs = opr.outputs
  362. self._compile()
  363. def get_opr_by_type(self, oprcls, unique=True):
  364. assert issubclass(oprcls, OpNode)
  365. rst = self.opr_filter.type(oprcls).as_list()
  366. if unique:
  367. assert len(rst) == 1, "{} operators of type {} found".format(
  368. len(rst), oprcls
  369. )
  370. (rst,) = rst
  371. return rst
  372. def get_opr_by_name(self, name, unique=True):
  373. rst = self.opr_filter.name(name).as_list()
  374. if unique:
  375. assert len(rst) == 1, "{} operators of type {} found".format(len(rst), name)
  376. (rst,) = rst
  377. return rst
  378. def get_var_by_name(self, name, unique=True):
  379. rst = self.var_filter.name(name).as_list()
  380. if unique:
  381. assert len(rst) == 1, "{} operators of type {} found".format(len(rst), name)
  382. (rst,) = rst
  383. return rst
  384. def get_var_receive_oprs(self, var):
  385. r"""Gets all oprs which use var as input"""
  386. return self.opr_filter.has_input(var).as_list()
  387. def get_dep_oprs(self, var):
  388. r"""Gets dependent oprs of var"""
  389. return get_oprs_seq(var, False, False)
  390. @property
  391. def opr_filter(self):
  392. r"""Filter on all opnodes of the Network."""
  393. oprs = self.all_oprs
  394. return NodeFilter(itertools.islice(oprs, len(oprs)))
  395. @property
  396. def var_filter(self):
  397. r"""Filter on all varnode of the Network."""
  398. vars = self.all_vars
  399. return NodeFilter(itertools.islice(vars, len(vars)))
  400. @property
  401. def params_filter(self): # all immutable tensor
  402. r"""Filter on all parameters (ImmutableTensor Opr) of the Network"""
  403. return self.opr_filter.param_provider()
  404. @property
  405. def data_providers_filter(self): # all host2devicecopy
  406. r"""Filter on all input nodes (Host2DeviceCopy Opr) of the Network"""
  407. return self.opr_filter.data_provider()
  408. @property
  409. def dest_vars(self):
  410. r"""Output varnodes of the Network."""
  411. return self.output_vars
  412. @property
  413. def all_oprs(self):
  414. return get_oprs_seq(self.output_vars, False, False)
  415. @property
  416. def all_vars(self):
  417. return get_dep_vars(self.output_vars)
  418. @property
  419. def all_vars_dict(self):
  420. return self.var_filter.as_dict()
  421. @property
  422. def all_oprs_dict(self):
  423. return self.opr_filter.as_dict()
  424. def _add_opr(self, opr) -> Optional[OpNode]:
  425. r"""Used for loading and building graph."""
  426. assert isinstance(opr, _imperative_rt.graph.OperatorNode)
  427. # TODO: use megbrain C++ RTTI to replace type string
  428. if opr.id not in self.all_oprs_map:
  429. opnode = str_to_mge_class(get_opr_type(opr)).load(opr)
  430. self.all_oprs_map[opr.id] = opnode
  431. for var in opr.inputs:
  432. varnode = self._get_var(var)
  433. opnode.add_inp_var(varnode)
  434. varnode.users.append(opnode)
  435. for var in opr.outputs:
  436. opnode.add_out_var(self._get_var(var))
  437. return opnode
  438. else:
  439. # overwrite the opnode 'new' output VarNode with
  440. # original one when output number larger than 1,
  441. # or will cause dependence issue in _compiler step.
  442. if len(opr.outputs) > 1:
  443. opnode = self.all_oprs_map[opr.id]
  444. for idx, output in enumerate(opnode.outputs):
  445. if output.var.id in self.all_vars_map:
  446. opnode.outputs[idx] = self.all_vars_map[output.var.id]
  447. return None
  448. def _get_opr(self, x):
  449. if x.id in self.all_oprs_map:
  450. return self.all_oprs_map[x.id]
  451. else:
  452. return None
  453. def _get_var(self, x):
  454. r"""Convert :class:`~._imperative_rt.graph.VarNode` to :class:`~.VarNode`."""
  455. assert isinstance(x, _imperative_rt.graph.VarNode)
  456. if x.id not in self.all_vars_map or self.all_vars_map[x.id].var != x:
  457. self.all_vars_map[x.id] = VarNode.load(x, self._get_opr(x.owner))
  458. return self.all_vars_map[x.id]
  459. def set_symbolic_shape(option: bool):
  460. r"""Set the VarNode use symbolic shape or not, return the last status.
  461. Please set to True and must recover after dump if want to change the input batch size.
  462. Args:
  463. option: True for enable symbolic shape.
  464. """
  465. return _set_symbolic_shape(option)
  466. def as_varnode(obj):
  467. r"""convert a :class:`.utils.network_node.VarNode` compatible object to :class:`.utils.network_node.VarNode`.
  468. Args:
  469. obj: it must be one of the following:
  470. 1. a :class:`.utils.network_node.VarNode` object
  471. 2. a :class:`.utils.network_node.OpNode` object that has unique output
  472. 3. an iterable that produces either type 1 or 2, with length 1
  473. """
  474. if type(obj) is VarNode:
  475. return obj
  476. if isinstance(obj, OpNode):
  477. assert len(obj.outputs) == 1, (
  478. "operator {} must have one output to be converted to VarNode; "
  479. "got {} actually".format(obj, len(obj.outputs))
  480. )
  481. ret = obj.outputs[0]
  482. assert type(ret) is VarNode
  483. return ret
  484. assert isinstance(
  485. obj, collections.abc.Iterable
  486. ), "{} is not compatible with VarNode".format(obj)
  487. val = list(obj)
  488. assert (
  489. len(val) == 1
  490. ), "can not convert sequence of length {} to VarNode ({})".format(
  491. len(val), (lambda s: s if len(s) < 50 else s[:50] + " ...")(str(val))
  492. )
  493. return as_varnode(val[0])
  494. def as_oprnode(obj):
  495. r"""convert a :class:`.utils.network_node.OpNode` compatible object to
  496. :class:`.utils.network_node.OpNode`; it works like :func:`as_varnode`.
  497. """
  498. if type(obj) is VarNode:
  499. return obj.owner
  500. if isinstance(obj, OpNode):
  501. return obj
  502. assert isinstance(
  503. obj, collections.abc.Iterable
  504. ), "{} is not compatible with OpNode".format(obj)
  505. val = list(obj)
  506. assert (
  507. len(val) == 1
  508. ), "can not convert sequence of length {} to " "OpNode({})".format(len(val), val)
  509. return as_oprnode(val[0])
  510. class NodeFilter:
  511. r"""Filter on node iterator. This class is an iterator of
  512. :class:`.NetworkNode` objects and multiple filtering conditions and
  513. mappers can be chained.
  514. Example:
  515. .. code-block::
  516. # find all :class:`.ImmutableTensor` nodes
  517. for i in NodeFilter(node_iter).param_provider():
  518. print(i)
  519. # find all :class:`.ImmutableTensor` nodes that end with ':W'
  520. for i in NodeFilter(node_iter).param_provider().name('*:W'):
  521. print(i)
  522. # number of inputs
  523. nr_input = NodeFilter(node_iter).data_provider().as_count()
  524. """
  525. _iter = None
  526. def __init__(self, node_iter):
  527. """
  528. :param node_iter: iterator to :class:`.NetworkNode`, or a
  529. :class:`.VarNode`-compatible object; in the later case, its
  530. dependent oprs would be used
  531. """
  532. if isinstance(node_iter, VarNode):
  533. oprs = get_oprs_seq(node_iter, False, False)
  534. node_iter = itertools.islice(oprs, len(oprs) - 1)
  535. if isinstance(node_iter, OpNode):
  536. oprs = get_oprs_seq(node_iter.inputs, False, False)
  537. node_iter = itertools.islice(oprs, len(oprs) - 1)
  538. assert isinstance(node_iter, collections.abc.Iterable)
  539. if (not isinstance(node_iter, NodeFilter)) and type(
  540. self
  541. ) is not NodeFilterCheckType:
  542. node_iter = NodeFilterCheckType(node_iter, NetworkNode)
  543. self._iter = node_iter
  544. @classmethod
  545. def make_all_deps(cls, *dest_vars):
  546. r"""make a :class:`NodeFilter` that contains all deps of given vars"""
  547. return cls(list(get_oprs_seq(dest_vars, False, False)))
  548. def __iter__(self):
  549. r"""to be overwritten by subclass to implement filters"""
  550. return iter(self._iter)
  551. def type(self, node_type):
  552. r"""filter by specific node type
  553. Args:
  554. node_type: node type class
  555. Returns:
  556. a new :class:`NodeFilter` object
  557. """
  558. return NodeFilterType(self, node_type)
  559. def check_type(self, node_type):
  560. r"""assert that all oprs produced by this iterator are instances of
  561. certain type
  562. Args:
  563. node_type: node type class
  564. Returns:
  565. a new :class:`NodeFilter` object
  566. Raises:
  567. TypeError if type check failed
  568. """
  569. return NodeFilterCheckType(self, node_type)
  570. def not_type(self, node_type):
  571. r"""remove oprs of specific type
  572. Args:
  573. node_type: node type class
  574. Returns:
  575. a new :class:`NodeFilter` object
  576. """
  577. return NodeFilterNotType(self, node_type)
  578. def param_provider(self):
  579. r"""get :class:`~.ParamProvider` oprs; shorthand for
  580. ``.type(ParamProvider)``
  581. """
  582. return self.type(ImmutableTensor)
  583. def data_provider(self):
  584. r"""get :class:`.DataProvider` oprs; shorthand for
  585. ``.type(DataProvider)``
  586. """
  587. return self.type(Host2DeviceCopy)
  588. def name(self, pattern, ignorecase=True):
  589. r"""filter by node name
  590. Args:
  591. pattern(class:`str`): a string in glob syntax that can contain ``?`` and
  592. ``*`` to match a single or arbitrary characters.
  593. ignorecase(bool, optional): whether to ignroe case
  594. Returns:
  595. a new :class:`NodeFilter` object
  596. """
  597. return NodeFilterName(self, pattern, ignorecase)
  598. def has_input(self, var):
  599. r"""an opr is kept if it has given var as one of its inputs
  600. Args:
  601. var: var node to checked
  602. Returns:
  603. a new :class:`NodeFilter` object
  604. """
  605. return NodeFilterHasInput(self, var)
  606. def as_list(self):
  607. r"""consume this iterator and return its content as a list"""
  608. return list(self)
  609. def as_unique(self):
  610. r"""assert that this iterator yields only one node and return it
  611. Returns:
  612. class:`.GraphNodeBase`: the unique node
  613. Raises:
  614. ValueError if this iterator does not yield a unique node
  615. """
  616. (opr,) = self
  617. return opr
  618. def as_dict(self):
  619. r"""construct an ordered dict to map from node names to objects in
  620. this iterator
  621. """
  622. return collections.OrderedDict((i.name, i) for i in self)
  623. def as_count(self):
  624. r"""consume this iterator and get the number of elements"""
  625. return sum(1 for _ in self)
  626. class NodeFilterType(NodeFilter):
  627. r"""see :meth:`NodeFilter.type`"""
  628. _node_type = None
  629. def __init__(self, node_iter, node_type):
  630. assert issubclass(node_type, NetworkNode), "bad opr type: {}".format(node_type)
  631. super().__init__(node_iter)
  632. self._node_type = node_type
  633. def __iter__(self):
  634. for i in self._iter:
  635. if isinstance(i, self._node_type):
  636. yield i
  637. class NodeFilterNotType(NodeFilterType):
  638. r"""see :meth:`NodeFilter.not_type`"""
  639. def __iter__(self):
  640. for i in self._iter:
  641. if not isinstance(i, self._node_type):
  642. yield i
  643. class NodeFilterCheckType(NodeFilterType):
  644. r"""see :meth:`NodeFilter.check_type`"""
  645. def __iter__(self):
  646. for i in self._iter:
  647. if not isinstance(i, self._node_type):
  648. raise TypeError(
  649. "all nodes should be {}; got {!r}".format(self._node_type, i)
  650. )
  651. yield i
  652. class NodeFilterHasInput(NodeFilter):
  653. r"""see :meth:`NodeFilter.has_input`"""
  654. _var = None
  655. def __init__(self, node_iter, var):
  656. var = as_varnode(var)
  657. super().__init__(node_iter)
  658. self.var = var
  659. def __iter__(self):
  660. for i in self._iter:
  661. assert isinstance(
  662. i, OpNode
  663. ), "has_input() must be used with OpNode; " "got {!r}".format(i)
  664. if any(self.var is _ for _ in i.inputs):
  665. yield i
  666. class NodeFilterName(NodeFilter):
  667. r"""see :meth:`NodeFilter.name`"""
  668. _re = None
  669. def __init__(self, node_iter, pattern, ignorecase):
  670. super().__init__(node_iter)
  671. self.pattern = pattern
  672. self._re = self.make_re(pattern, ignorecase)
  673. @classmethod
  674. def make_re(cls, pattern, ignorecase=True):
  675. assert isinstance(pattern, str), "bad pattern: {!r}".format(pattern)
  676. assert isinstance(ignorecase, bool)
  677. flags = 0
  678. if ignorecase:
  679. flags |= re.IGNORECASE
  680. return re.compile(fnmatch.translate(pattern), flags=flags)
  681. def __iter__(self):
  682. for i in self._iter:
  683. if self.pattern == i.name or self._re.match(i.name):
  684. yield i