Skip to content

Commit

Permalink
Add shape test
Browse files Browse the repository at this point in the history
Add two node tests
  • Loading branch information
wschin committed Apr 29, 2019
1 parent 17c35f5 commit fd1568d
Show file tree
Hide file tree
Showing 26 changed files with 228 additions and 11 deletions.
85 changes: 81 additions & 4 deletions docs/TestCoverage.md
Expand Up @@ -5,7 +5,7 @@
* [Overall Test Coverage](#overall-test-coverage)
# Node Test Coverage
## Summary
Node tests have covered 125/133 (93.98%, 5 generators excluded) common operators.
Node tests have covered 126/133 (94.74%, 5 generators excluded) common operators.

Node tests have covered 0/0 (N/A) experimental operators.

Expand Down Expand Up @@ -88,6 +88,86 @@ expect(node, inputs=[x], outputs=[y],
</details>


### Adagrad
There are 2 test cases, listed as following:
<details>
<summary>adagrad</summary>

```python
# Define operator attributes.
norm_coefficient = 0.001
epsilon = 1e-5
decay_factor = 0.1

# Create operator.
node = onnx.helper.make_node('Adagrad',
inputs=['R', 'T', 'X', 'G', 'H'],
outputs=['X_new, H_new'],
norm_coefficient=norm_coefficient,
epsilon=epsilon,
decay_factor=decay_factor
)

# Define operator inputs.
r = np.array(0.1, dtype=np.float32) # scalar
t = np.array(0, dtype=np.int64) # scalar
x = np.array([1.0], dtype=np.float32)
g = np.array([-1.0], dtype=np.float32)
h = np.array([2.0], dtype=np.float32)

# Compute expected outputs of Adagrad.
x_new, h_new = apply_adagrad(r, t, x, g, h,
norm_coefficient, epsilon, decay_factor)

# Check results.
expect(node, inputs=[r, t, x, g, h],
outputs=[x_new, h_new], name='test_adagrad')
```

</details>
<details>
<summary>adagrad_multiple</summary>

```python
# Define operator attributes.
norm_coefficient = 0.001
epsilon = 1e-5
decay_factor = 0.1

node = onnx.helper.make_node('Adagrad',
inputs=['R', 'T', 'X', 'G', 'H'],
outputs=['X_new, H_new'],
norm_coefficient=norm_coefficient,
epsilon=epsilon,
decay_factor=decay_factor
)

# Define operator inputs.
r = np.array(0.1, dtype=np.float32) # scalar
t = np.array(0, dtype=np.int64) # scalar

x1 = np.array([1.0], dtype=np.float32)
g1 = np.array([-1.0], dtype=np.float32)
h1 = np.array([2.0], dtype=np.float32)

x2 = np.array([1.0, 2.0], dtype=np.float32)
g2 = np.array([-1.0, -3.0], dtype=np.float32)
h2 = np.array([4.0, 1.0], dtype=np.float32)

# Compute expected outputs of Adagrad.
x1_new, h1_new = apply_adagrad(r, t, x1, g1, h1,
norm_coefficient, epsilon, decay_factor)
x2_new, h2_new = apply_adagrad(r, t, x2, g2, h2,
norm_coefficient, epsilon, decay_factor)

# Check results.
expect(node, inputs=[r, t, x1, x2, g1, g2, h1, h2],
outputs=[x1_new, x2_new, h1_new, h2_new], name='test_adagrad_multiple')
```

</details>


### Add
There are 2 test cases, listed as following:
<details>
Expand Down Expand Up @@ -7921,9 +8001,6 @@ expect(node, inputs=[x, y], outputs=[z],
<br/>

## &#x1F494;No Cover Common Operators
### Adagrad (call for test cases)


### GlobalLpPool (call for test cases)


Expand Down
95 changes: 95 additions & 0 deletions onnx/backend/test/case/node/adagrad.py
@@ -0,0 +1,95 @@
from __future__ import absolute_import
from __future__ import division
from __future__ import print_function
from __future__ import unicode_literals

import numpy as np # type: ignore

import onnx
from ..base import Base
from . import expect

def apply_adagrad(r, t, x, g, h, norm_coefficient, epsilon, decay_factor):
# Compute adjusted learning-rate.
r_ = r * (1 + t * decay_factor)
# Add gradient of regularization term.
g_regularized = norm_coefficient * x + g
# Update squared accumulated gradient.
h_new = h + g * g
# Compute ADAGRAD's gradient scaling factors
h_sqrt = np.sqrt(h_new) + epsilon
# Apply ADAGRAD update rule.
x_new = x - r_ * g_regularized / h_sqrt
return (x_new, h_new)

class Adagrad(Base):

@staticmethod
def export_adagrad(): # type: () -> None
# Define operator attributes.
norm_coefficient = 0.001
epsilon = 1e-5
decay_factor = 0.1

# Create operator.
node = onnx.helper.make_node('Adagrad',
inputs=['R', 'T', 'X', 'G', 'H'],
outputs=['X_new', 'H_new'],
norm_coefficient=norm_coefficient,
epsilon=epsilon,
decay_factor=decay_factor
)

# Define operator inputs.
r = np.array(0.1, dtype=np.float32) # scalar
t = np.array(0, dtype=np.int64) # scalar
x = np.array([1.0], dtype=np.float32)
g = np.array([-1.0], dtype=np.float32)
h = np.array([2.0], dtype=np.float32)

# Compute expected outputs of Adagrad.
x_new, h_new = apply_adagrad(r, t, x, g, h,
norm_coefficient, epsilon, decay_factor)

# Check results.
expect(node, inputs=[r, t, x, g, h],
outputs=[x_new, h_new], name='test_adagrad')

@staticmethod
def export_adagrad_multiple(): # type: () -> None
# Define operator attributes.
norm_coefficient = 0.001
epsilon = 1e-5
decay_factor = 0.1

node = onnx.helper.make_node('Adagrad',
inputs=['R', 'T', 'X1', 'X2',
'G1', 'G2', 'H1', 'H2'],
outputs=['X1_new', 'X2_new',
'H1_new', 'H2_new'],
norm_coefficient=norm_coefficient,
epsilon=epsilon,
decay_factor=decay_factor
)

# Define operator inputs.
r = np.array(0.1, dtype=np.float32) # scalar
t = np.array(0, dtype=np.int64) # scalar

x1 = np.array([1.0], dtype=np.float32)
g1 = np.array([-1.0], dtype=np.float32)
h1 = np.array([2.0], dtype=np.float32)

x2 = np.array([1.0, 2.0], dtype=np.float32)
g2 = np.array([-1.0, -3.0], dtype=np.float32)
h2 = np.array([4.0, 1.0], dtype=np.float32)

# Compute expected outputs of Adagrad.
x1_new, h1_new = apply_adagrad(r, t, x1, g1, h1,
norm_coefficient, epsilon, decay_factor)
x2_new, h2_new = apply_adagrad(r, t, x2, g2, h2,
norm_coefficient, epsilon, decay_factor)

# Check results.
expect(node, inputs=[r, t, x1, x2, g1, g2, h1, h2],
outputs=[x1_new, x2_new, h1_new, h2_new], name='test_adagrad_multiple')
Binary file added onnx/backend/test/data/node/test_adagrad/model.onnx
Binary file not shown.
@@ -0,0 +1 @@
BRJ���=
Binary file not shown.
Binary file not shown.
Binary file not shown.
Binary file not shown.
@@ -0,0 +1 @@
BX_newJ�a�?
Binary file not shown.
Binary file not shown.
@@ -0,0 +1 @@
BRJ���=
Binary file not shown.
Binary file not shown.
Binary file not shown.
Binary file not shown.
Binary file not shown.
Binary file not shown.
Binary file not shown.
@@ -0,0 +1 @@
BX1_newJ�a�?
@@ -0,0 +1 @@
BX2_newJ���?H@
Binary file not shown.
Binary file not shown.
2 changes: 1 addition & 1 deletion onnx/defs/operator_sets.h
Expand Up @@ -576,7 +576,7 @@ class OpSet_Onnx_ver10 {
fn(GetOpSchema<ONNX_OPERATOR_SET_SCHEMA_CLASS_NAME(
Onnx, 10, ReverseSequence)>());
fn(GetOpSchema<ONNX_OPERATOR_SET_SCHEMA_CLASS_NAME(
Onnx, 10, RoiAlign)>());
Onnx, 10, RoiAlign)>());
fn(GetOpSchema<ONNX_OPERATOR_SET_SCHEMA_CLASS_NAME(
Onnx, 10, Adagrad)>());
}
Expand Down
19 changes: 13 additions & 6 deletions onnx/defs/optimizer/defs.cc
Expand Up @@ -13,7 +13,7 @@ static const char* Adagrad_ver10_doc = R"DOC(
- The initial learning-rate "R".
- The update count "T". That is, the number of training iterations conducted.
- A L2-norm regularization coefficient "lambda".
- A L2-norm regularization coefficient "norm_coefficient".
- A learning-rate decay factor "decay_factor".
- A small constant "epsilon" to avoid dividing-by-zero.
Expand All @@ -33,8 +33,8 @@ static const char* Adagrad_ver10_doc = R"DOC(
// Compute a scalar learning-rate factor. If X is never updated, T should be 0.
r = R / (1 + T * decay_factor);
// Add gradient of 0.5 * lambda * ||X||_2^2, where ||X||_2 is the 2-norm.
G_regularized = lambda * X + G;
// Add gradient of 0.5 * norm_coefficient * ||X||_2^2, where ||X||_2 is the 2-norm.
G_regularized = norm_coefficient * X + G;
// Compute new accumulated squared gradient.
H_new = H + G_regularized * G_regularized;
Expand Down Expand Up @@ -96,8 +96,8 @@ ONNX_OPERATOR_SET_SCHEMA(
AttributeProto::FLOAT,
0.0f)
.Attr(
"lambda",
"Regularization coefficient of 0.5 * lambda * ||X||_2^2. Default to 0, "
"norm_coefficient",
"Regularization coefficient in 0.5 * norm_coefficient * ||X||_2^2. Default to 0, "
"which means no regularization.",
AttributeProto::FLOAT,
0.0f)
Expand All @@ -112,5 +112,12 @@ ONNX_OPERATOR_SET_SCHEMA(
.TypeConstraint(
"T3",
{"tensor(float)", "tensor(double)"},
"Constrain input types to float tensors."));
"Constrain input types to float tensors.")
.TypeAndShapeInferenceFunction([](InferenceContext& ctx) {
auto num_inputs = ctx.getNumInputs();
// skip 'R' and 'T' in input list and then propoagate other shapes to outputs.
for (size_t i = 2; i < num_inputs; ++i) {
propagateElemTypeFromInputToOutput(ctx, i, i - 2);
propagateShapeFromInputToOutput(ctx, i, i - 2);
}}));
} // namespace ONNX_NAMESPACE
33 changes: 33 additions & 0 deletions onnx/test/shape_inference_test.py
Expand Up @@ -1694,6 +1694,39 @@ def test_reversesequence(self): # type: () -> None
[])
self._assert_inferred(graph, [make_tensor_value_info('y', TensorProto.FLOAT, (4, 5, 6))])

def test_adagrad(self): # type: () -> None
graph = self.make_graph(
[('R', TensorProto.FLOAT, ()), # scalar's shape is ()
('T', TensorProto.INT64, ()), # scalar's shape is ()
('X', TensorProto.FLOAT, (1, 2)),
('G', TensorProto.FLOAT, (1, 2)),
('H', TensorProto.FLOAT, (1, 2))],
[make_node('Adagrad', ['R', 'T', 'X', 'G', 'H'], ['X_new', 'H_new'])],
[])
self._assert_inferred(graph,
[make_tensor_value_info('X_new', TensorProto.FLOAT, (1, 2)),
make_tensor_value_info('H_new', TensorProto.FLOAT, (1, 2))])

def test_adagrad_multiple(self): # type: () -> None
graph = self.make_graph(
[('R', TensorProto.FLOAT, ()), # scalar's shape is ()
('T', TensorProto.INT64, ()), # scalar's shape is ()
('X1', TensorProto.FLOAT, (1, 2)),
('X2', TensorProto.FLOAT, (3, 4)),
('G1', TensorProto.FLOAT, (1, 2)),
('G2', TensorProto.FLOAT, (3, 4)),
('H1', TensorProto.FLOAT, (1, 2))
('H2', TensorProto.FLOAT, (3, 4))],
[make_node('Adagrad', ['R', 'T', 'X1', 'X2', 'G1', 'G2', 'H1', 'H2'],
['X1_new', 'X2_new', 'H1_new', 'H2_new'])],
[])
self._assert_inferred(graph,
[make_tensor_value_info('X1_new', TensorProto.FLOAT, (1, 2)),
make_tensor_value_info('X2_new', TensorProto.FLOAT, (3, 4)),
make_tensor_value_info('H1_new', TensorProto.FLOAT, (1, 2)),
make_tensor_value_info('H2_new', TensorProto.FLOAT, (3, 4))
])


if __name__ == '__main__':
unittest.main()

0 comments on commit fd1568d

Please sign in to comment.