iter.py 9.6 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266
  1. # mypy: ignore-errors
  2. MAX_CYCLE = 3000
  3. import itertools
  4. import operator
  5. from typing import Dict, List, Optional
  6. from .. import polyfill, variables
  7. from ..exc import unimplemented
  8. from .base import MutableLocal, VariableTracker
  9. from .constant import ConstantVariable
  10. class ItertoolsVariable(VariableTracker):
  11. def __init__(self, value, **kwargs):
  12. super().__init__(**kwargs)
  13. self.value = value
  14. def __repr__(self):
  15. return f"ItertoolsVariable({self.value})"
  16. def python_type(self):
  17. return type(self.value)
  18. def as_python_constant(self):
  19. return self.value
  20. def call_function(
  21. self, tx, args: "List[VariableTracker]", kwargs: "Dict[str, VariableTracker]"
  22. ) -> "VariableTracker":
  23. if (
  24. self.value is itertools.product
  25. and not kwargs
  26. and all(arg.has_unpack_var_sequence(tx) for arg in args)
  27. ):
  28. seqs = [arg.unpack_var_sequence(tx) for arg in args]
  29. items = []
  30. for item in itertools.product(*seqs):
  31. items.append(variables.TupleVariable(list(item)))
  32. return variables.ListIteratorVariable(items, mutable_local=MutableLocal())
  33. elif (
  34. self.value is itertools.chain
  35. and not kwargs
  36. and all(arg.has_unpack_var_sequence(tx) for arg in args)
  37. ):
  38. seqs = [arg.unpack_var_sequence(tx) for arg in args]
  39. items = list(itertools.chain.from_iterable(seqs))
  40. return variables.ListIteratorVariable(items, mutable_local=MutableLocal())
  41. elif self.value is itertools.accumulate:
  42. from .builtin import BuiltinVariable
  43. if any(key not in ["initial", "func"] for key in kwargs.keys()):
  44. unimplemented(
  45. "Unsupported kwargs for itertools.accumulate: "
  46. f"{','.join(set(kwargs.keys()) - {'initial', 'func'})}"
  47. )
  48. acc = kwargs.get("initial")
  49. if len(args) in [1, 2] and args[0].has_unpack_var_sequence(tx):
  50. seq = args[0].unpack_var_sequence(tx)
  51. if "func" in kwargs and len(args) == 1:
  52. func = kwargs["func"].call_function
  53. elif len(args) == 2:
  54. func = args[1].call_function
  55. elif len(args) == 1:
  56. # Default to operator.add
  57. func = BuiltinVariable(operator.add).call_function
  58. else:
  59. unimplemented(
  60. "itertools.accumulate can only accept one of: `func` kwarg, pos 2 arg"
  61. )
  62. else:
  63. unimplemented("Unsupported arguments for itertools.accumulate")
  64. items = []
  65. if acc is not None:
  66. items.append(acc)
  67. for item in seq:
  68. if acc is None:
  69. acc = item
  70. else:
  71. try:
  72. acc = func(tx, [acc, item], {})
  73. except Exception as e:
  74. unimplemented(
  75. f"Unexpected failure in invoking function during accumulate. Failed running func {func}({item}{acc})",
  76. from_exc=e,
  77. )
  78. items.append(acc)
  79. return variables.ListIteratorVariable(items, mutable_local=MutableLocal())
  80. elif (
  81. self.value is itertools.combinations
  82. and not kwargs
  83. and len(args) == 2
  84. and args[0].has_unpack_var_sequence(tx)
  85. and args[1].is_python_constant()
  86. ):
  87. iterable = args[0].unpack_var_sequence(tx)
  88. r = args[1].as_python_constant()
  89. items = []
  90. for item in itertools.combinations(iterable, r):
  91. items.append(variables.TupleVariable(list(item)))
  92. return variables.ListIteratorVariable(items, mutable_local=MutableLocal())
  93. elif self.value is itertools.groupby:
  94. if any(kw != "key" for kw in kwargs.keys()):
  95. unimplemented(
  96. "Unsupported kwargs for itertools.groupby: "
  97. f"{','.join(set(kwargs.keys()) - {'key'})}"
  98. )
  99. def retrieve_const_key(key):
  100. if isinstance(key, variables.SymNodeVariable):
  101. return key.evaluate_expr()
  102. elif isinstance(key, variables.ConstantVariable):
  103. return key.as_python_constant()
  104. else:
  105. unimplemented(
  106. "Unsupported key type for itertools.groupby: " + str(type(key))
  107. )
  108. if len(args) == 1 and args[0].has_unpack_var_sequence(tx):
  109. seq = args[0].unpack_var_sequence(tx)
  110. keyfunc = (
  111. (
  112. lambda x: (
  113. retrieve_const_key(
  114. kwargs.get("key").call_function(tx, [x], {})
  115. )
  116. )
  117. )
  118. if "key" in kwargs
  119. else None
  120. )
  121. else:
  122. unimplemented("Unsupported arguments for itertools.groupby")
  123. result = []
  124. try:
  125. for k, v in itertools.groupby(seq, key=keyfunc):
  126. result.append(
  127. variables.TupleVariable(
  128. [
  129. variables.ConstantVariable.create(k)
  130. if variables.ConstantVariable.is_literal(k)
  131. else k,
  132. variables.ListIteratorVariable(
  133. list(v), mutable_local=MutableLocal()
  134. ),
  135. ],
  136. mutable_local=MutableLocal(),
  137. )
  138. )
  139. except Exception as e:
  140. unimplemented(
  141. "Unexpected failure when calling itertools.groupby",
  142. from_exc=e,
  143. )
  144. return variables.ListIteratorVariable(result, mutable_local=MutableLocal())
  145. elif self.value is itertools.repeat:
  146. if len(args) < 2:
  147. return variables.RepeatIteratorVariable(
  148. *args, mutable_local=MutableLocal()
  149. )
  150. from .builder import SourcelessBuilder
  151. return tx.inline_user_function_return(
  152. SourcelessBuilder.create(tx, polyfill.repeat), args, kwargs
  153. )
  154. elif self.value is itertools.count:
  155. return variables.CountIteratorVariable(*args, mutable_local=MutableLocal())
  156. elif self.value is itertools.cycle:
  157. return variables.CycleIteratorVariable(*args, mutable_local=MutableLocal())
  158. elif self.value is itertools.dropwhile:
  159. return variables.UserFunctionVariable(polyfill.dropwhile).call_function(
  160. tx, args, kwargs
  161. )
  162. else:
  163. return super().call_function(tx, args, kwargs)
  164. class IteratorVariable(VariableTracker):
  165. def __init__(self, **kwargs):
  166. super().__init__(**kwargs)
  167. def next_variable(self, tx):
  168. unimplemented("abstract method, must implement")
  169. class RepeatIteratorVariable(IteratorVariable):
  170. def __init__(self, item: VariableTracker, **kwargs):
  171. super().__init__(**kwargs)
  172. self.item = item
  173. # Repeat needs no mutation, clone self
  174. def next_variable(self, tx):
  175. return self.item
  176. class CountIteratorVariable(IteratorVariable):
  177. def __init__(self, item: int = 0, step: int = 1, **kwargs):
  178. super().__init__(**kwargs)
  179. if not isinstance(item, VariableTracker):
  180. item = ConstantVariable.create(item)
  181. if not isinstance(step, VariableTracker):
  182. step = ConstantVariable.create(step)
  183. self.item = item
  184. self.step = step
  185. def next_variable(self, tx):
  186. assert self.mutable_local
  187. tx.output.side_effects.mutation(self)
  188. next_item = self.item.call_method(tx, "__add__", [self.step], {})
  189. self.item = next_item
  190. return self.item
  191. class CycleIteratorVariable(IteratorVariable):
  192. def __init__(
  193. self,
  194. iterator: IteratorVariable,
  195. saved: List[VariableTracker] = None,
  196. saved_index: int = 0,
  197. item: Optional[VariableTracker] = None,
  198. **kwargs,
  199. ):
  200. if saved is None:
  201. saved = []
  202. super().__init__(**kwargs)
  203. self.iterator = iterator
  204. self.saved = saved
  205. self.saved_index = saved_index
  206. self.item = item
  207. def next_variable(self, tx):
  208. assert self.mutable_local
  209. if self.iterator is not None:
  210. try:
  211. new_item = self.iterator.next_variable(tx)
  212. if len(self.saved) > MAX_CYCLE:
  213. unimplemented(
  214. "input iterator to itertools.cycle has too many items"
  215. )
  216. tx.output.side_effects.mutation(self)
  217. self.saved.append(new_item)
  218. self.item = new_item
  219. if self.item is None:
  220. return self.next_variable(tx)
  221. return self.item
  222. except StopIteration:
  223. self.iterator = None
  224. return self.next_variable(tx)
  225. elif len(self.saved) > 0:
  226. tx.output.side_effects.mutation(self)
  227. self.saved_index = (self.saved_index + 1) % len(self.saved)
  228. return self.item
  229. else:
  230. raise StopIteration