Skip to content

Commit 681cb1d

Browse files
committed
fix np.cast
Signed-off-by: xadupre <[email protected]>
1 parent e4c3409 commit 681cb1d

File tree

5 files changed

+48
-45
lines changed

5 files changed

+48
-45
lines changed

.github/actions/keras_application_test/action.yml

Lines changed: 2 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -46,6 +46,7 @@ runs:
4646
pip install git+https://github.com/qubvel/efficientnet
4747
pip install keras-self-attention
4848
pip install pytest pytest-cov pytest-runner
49+
pip uninstall -y protobuf
4950
5051
if [[ ${{ inputs.tf_version }} == 1.* ]]; then
5152
pip install keras==2.3.1
@@ -60,7 +61,7 @@ runs:
6061
pip install tf_keras==${{ inputs.tf_version }} tensorflow==${{ inputs.tf_version }}
6162
else
6263
echo "-- install-1 TF ${{ inputs.tf_version }}"
63-
pip install "protobuf<3.20" tensorflow==${{ inputs.tf_version }}
64+
pip install protobuf tensorflow==${{ inputs.tf_version }}
6465
fi
6566
fi
6667

.github/actions/keras_unit_test/action.yml

Lines changed: 1 addition & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -39,7 +39,7 @@ runs:
3939
pip install tf_keras==${{ inputs.tf_version }} tensorflow==${{ inputs.tf_version }}
4040
else
4141
echo "-- install-2 TF ${{ inputs.tf_version }}"
42-
pip install "protobuf<3.20" tensorflow==${{ inputs.tf_version }}
42+
pip install protobuf tensorflow==${{ inputs.tf_version }}
4343
fi
4444
fi
4545

tests/test_custom_rnncell.py

Lines changed: 40 additions & 40 deletions
Original file line numberDiff line numberDiff line change
@@ -16,6 +16,7 @@
1616
# pylint: disable=abstract-method,arguments-differ
1717

1818
if is_tf2():
19+
# no test for tf2 in this file
1920
pass
2021
else:
2122
LSTMBlockCell = tf.contrib.rnn.LSTMBlockCell
@@ -26,6 +27,45 @@
2627
dynamic_rnn = tf.nn.dynamic_rnn
2728
bidirectional_dynamic_rnn = tf.nn.bidirectional_dynamic_rnn
2829

30+
class GatedGRUCell(RNNCell):
31+
def __init__(self, hidden_dim, reuse=None):
32+
super().__init__(self, _reuse=reuse)
33+
self._num_units = hidden_dim
34+
self._activation = tf.tanh
35+
36+
@property
37+
def state_size(self):
38+
return self._num_units
39+
40+
@property
41+
def output_size(self):
42+
return self._num_units
43+
44+
def call(self, inputs, state):
45+
# inputs shape: [batch size, time step, input size] = [1, 3, 2]
46+
# num_units: 5
47+
# W shape: [2, 3 * 5] = [2, 15]
48+
# U shape: [5, 3 * 5] = [5, 15]
49+
# b shape: [1, 3 * 5] = [1, 15]
50+
# state shape: [batch size, state size] = [1, 5]
51+
52+
input_dim = inputs.get_shape()[-1]
53+
assert input_dim is not None, "input dimension must be defined"
54+
# W = tf.get_variable(name="W", shape=[input_dim, 3 * self._num_units], dtype=tf.float32)
55+
W = np.arange(30.0, dtype=np.float32).reshape((2, 15))
56+
# U = tf.get_variable(name='U', shape=[self._num_units, 3 * self._num_units], dtype=tf.float32)
57+
U = np.arange(75.0, dtype=np.float32).reshape((5, 15))
58+
# b = tf.get_variable(name='b', shape=[1, 3 * self._num_units], dtype=tf.float32)
59+
b = np.arange(15.0, dtype=np.float32).reshape((1, 15))
60+
61+
xw = tf.split(tf.matmul(inputs, W) + b, 3, 1)
62+
hu = tf.split(tf.matmul(state, U), 3, 1)
63+
r = tf.sigmoid(xw[0] + hu[0])
64+
z = tf.sigmoid(xw[1] + hu[1])
65+
h1 = self._activation(xw[2] + r * hu[2])
66+
next_h = h1 * (1 - z) + state * z
67+
return next_h, next_h
68+
2969

3070
class CustomRnnCellTests(Tf2OnnxBackendTestBase):
3171
@check_opset_min_version(8, "Scan")
@@ -370,45 +410,5 @@ def func(encoder_x, decoder_x, seq_length):
370410
self.run_test_case(func, feed_dict, input_names_with_port, output_names_with_port, 0.1)
371411

372412

373-
class GatedGRUCell(RNNCell):
374-
def __init__(self, hidden_dim, reuse=None):
375-
super().__init__(self, _reuse=reuse)
376-
self._num_units = hidden_dim
377-
self._activation = tf.tanh
378-
379-
@property
380-
def state_size(self):
381-
return self._num_units
382-
383-
@property
384-
def output_size(self):
385-
return self._num_units
386-
387-
def call(self, inputs, state):
388-
# inputs shape: [batch size, time step, input size] = [1, 3, 2]
389-
# num_units: 5
390-
# W shape: [2, 3 * 5] = [2, 15]
391-
# U shape: [5, 3 * 5] = [5, 15]
392-
# b shape: [1, 3 * 5] = [1, 15]
393-
# state shape: [batch size, state size] = [1, 5]
394-
395-
input_dim = inputs.get_shape()[-1]
396-
assert input_dim is not None, "input dimension must be defined"
397-
# W = tf.get_variable(name="W", shape=[input_dim, 3 * self._num_units], dtype=tf.float32)
398-
W = np.arange(30.0, dtype=np.float32).reshape((2, 15))
399-
# U = tf.get_variable(name='U', shape=[self._num_units, 3 * self._num_units], dtype=tf.float32)
400-
U = np.arange(75.0, dtype=np.float32).reshape((5, 15))
401-
# b = tf.get_variable(name='b', shape=[1, 3 * self._num_units], dtype=tf.float32)
402-
b = np.arange(15.0, dtype=np.float32).reshape((1, 15))
403-
404-
xw = tf.split(tf.matmul(inputs, W) + b, 3, 1)
405-
hu = tf.split(tf.matmul(state, U), 3, 1)
406-
r = tf.sigmoid(xw[0] + hu[0])
407-
z = tf.sigmoid(xw[1] + hu[1])
408-
h1 = self._activation(xw[2] + r * hu[2])
409-
next_h = h1 * (1 - z) + state * z
410-
return next_h, next_h
411-
412-
413413
if __name__ == '__main__':
414414
unittest_main()

tests/utils/setup_test_env.sh

Lines changed: 4 additions & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -17,16 +17,18 @@ echo "==== ONNX version: $ONNX_VERSION"
1717

1818
pip install pytest pytest-cov pytest-runner coverage graphviz requests pyyaml pillow pandas parameterized sympy coloredlogs flatbuffers timeout-decorator
1919
pip uninstall -y tensorflow
20-
pip install onnx==$ONNX_VERSION onnxruntime==$ORT_VERSION onnxruntime-extensions "tensorflow-text<=$TF_VERSION" tensorflow==$TF_VERSION
20+
pip install onnx==$ONNX_VERSION onnxruntime==$ORT_VERSION onnxruntime-extensions
2121

2222
if [[ $TF_VERSION == 1.* ]]; then
2323
pip install numpy==1.19.0
2424
else
25+
pip uninstall -y protobuf
2526
if [[ "$TF_VERSION" != "2.13.0" && "$TF_VERSION" != "2.9.0" ]]; then
2627
echo "-- install-3 TF-KERAS $TF_VERSION"
27-
pip install tf_keras==$TF_VERSION
28+
pip install tf_keras==$TF_VERSION "tensorflow-text<=$TF_VERSION"
2829
else
2930
echo "-- install-3 TF $TF_VERSION"
31+
pip install "tensorflow-text<=$TF_VERSION" tensorflow==$TF_VERSION protobuf
3032
fi
3133
fi
3234

tf2onnx/tfonnx.py

Lines changed: 1 addition & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -74,7 +74,7 @@ def rewrite_constant_fold(g, ops):
7474
func_map = {
7575
"Add": np.add,
7676
"GreaterEqual": np.greater_equal,
77-
"Cast": np.cast,
77+
"Cast": lambda x, dtype: x.astype(dtype),
7878
"ConcatV2": np.concatenate,
7979
"Less": np.less,
8080
"ListDiff": np.setdiff1d,

0 commit comments

Comments
 (0)