Def _slice from Relay Tensorflow report has no attribute name_hint error

Hi:

I am trying to convert a SSD model from the Object Detection of Tensorflow Models into tvm. And I got an error <class 'tvm.relay.expr.Call'> has no attribute name_hint when converting a Slice operation.

It seems that one input node of the Slice operation is an ‘ExpandDims’ operation, which is converted into tvm.relay.expr.Call object. And when the def _slice() function for converting Slice operation calsl _get_list_param to get size, it crashes, since _get_list_param can not handle ``tvm.relay.expr.Call` object

So, any suggestion to fix such problems? I can only think of adjusting the Tensorflow graph_def before relay.frontend.from_tensorflow by replacing the ExpandDims nodes with Const nodes? Is there some simpler way?

Thanks a lot

Recently, we’ve been solving this by inferring the value of ExpandDims by running the graph up until that point and using that returned value as the input. As an example, check out how we get axes in _transpose:

try:
    axes = _get_list_param(params, inputs[1])
except (IndexError, KeyError, AttributeError):
    axes = _infer_value_simulated(inputs[1], params).asnumpy()

Feel free to send a PR if this fix works!

Thanks a lot. I will definitely try it.

Hi jonso:

I have changed

def _get_param(params, input_node):
    if isinstance(input_node, _expr.Constant):
        return np.atleast_1d(input_node.data.asnumpy())
    return params.pop(input_node.name_hint).asnumpy()

into

def _get_param(params, input_node):
    if isinstance(input_node, _expr.Constant):
        return np.atleast_1d(input_node.data.asnumpy())
    elif isinstance(input_node, _expr.Call):
        return _infer_value_simulated(input_node, params).asnumpy()[0]
    return params.pop(input_node.name_hint).asnumpy()

I have changed the def _get_param function since def _slice seems to be not the only operation need to handle tvm.relay.expr.Call input. Although the error before seems to be fixed, another bug comes out. When a reshape node with two inputs, which are an add operation node and a concatv2 node, is converted, there comes out antoher error info:

Check failed: begin_v < end_v (0 vs. 0) : strided_slice get empty slice at axis 0

This is the url of the graph_def I want to convert, which can be loaded by

def load_graph_def(graph_def_path):
    infer_graph_def = tf.GraphDef()
    with open(graph_def_path, 'rb') as f:
        infer_graph_def.ParseFromString(f.read())
    return infer_graph_def

Thanks.

@lsy643 I met same issue and change _get_param to solve ‘<class ‘tvm.relay.expr.Call’> has no attribute name_hint’ error.

another bug also comes out. Have you solved this problem?

#StridedSlice op
def _stridedSlice():
    def _impl(inputs, attr, params, mod):
        xxxxx
        out = _op.strided_slice(inputs[0], begin=begin, end=end, strides=stride)
        out_shape = _infer_shape(out, mod=mod)   # <--- error


failed: begin_v < end_v (3 vs. 0) : strided_slice get empty slice at axis 0
......
tvm::IRModuleNode::Add(xxx)
tvm::RunTypeCheck(xx)

Perhaps one attr of _impl returns an empty array. But I remember they have solved this problem. Try to get the latest TVM code and run it again.

If it still not works, you can try to add something like this into the def _impl before _op.strided_slice

        # handle situations when return empty constant
        input0_np = _infer_value(inputs[0], params).asnumpy()  
        input0_shape = input0_np.shape
        for input_dim_i, begin_i, size_i in zip(input0_shape, begin, size):
            # when one size_i == 0, return empty constant
            if size_i == 0:
                return _expr.const([], dtype=input0_np.dtype) 
            # when begin index larger than the dimension of inputs[0]
            if begin_i >= input_dim_i:
                return _expr.const([], dtype=input0_np.dtype) 

@lsy643 I get the latest TVM code and error.

My error is in another function (_stridedSlice) , is not the _slice.

def _stridedSlice():
    def _impl(inputs, attr, params, mod):
        """Strided Slice.
        Operator description: https://www.tensorflow.org/api_docs/python/tf/strided_slice
        Tensorflow mask validation: https://github.com/tensorflow/tensorflow/blob/master/
        tensorflow/core/util/strided_slice_op.cc#L147-L368
        """
        begin = _get_list_param(params, inputs[1])
        end = _get_list_param(params, inputs[2])
        stride = _get_list_param(params, inputs[3])
        begin_mask = int(attr.get('begin_mask', 0))
        end_mask = int(attr.get('end_mask', 0))
        ellipsis_mask = int(attr.get('ellipsis_mask', 0))
        new_axis_mask = int(attr.get('new_axis_mask', 0))
        shrink_axis_mask = int(attr.get('shrink_axis_mask', 0))
        data_shape = attr['_input_shapes'][inputs[0]]
        data_dim = len(data_shape)
        stride_dim = len(stride)

        def _transform_mask(stride_dim, ellipsis_mask):
            """Handle mask inputs to create new begin, end, stride and output shape"""
            m_begin = [0] * data_dim
            m_end = [0] * data_dim
            m_stride = [0] * data_dim
            fshape_indices = []
            #Count new axis after ellipsis_mask, consider while applying ellipsis_mask.
            ellipsis_seen = False
            new_axes_after_ellipsis = 0
            for i in range(stride_dim):
                mask = 1 << i
                if ellipsis_seen and (mask & new_axis_mask) != 0:
                    new_axes_after_ellipsis += 1
                if (mask & ellipsis_mask) != 0:
                    ellipsis_seen = True
            if not ellipsis_seen:
                #Used later for extending the stride attributes in the below loop.
                ellipsis_mask |= (1 << stride_dim)
                stride_dim += 1
            final_index = 0
            for index in range(stride_dim):
                mask = 1 << index
                if mask & ellipsis_mask:
                    #Identify the end index for applying ellipsis_mask
                    to_index = min(((data_dim - (stride_dim-index)) + 1 \
                                     + new_axes_after_ellipsis), data_dim)
                    for i in range(final_index, to_index):
                        m_begin[final_index] = 0
                        m_end[final_index] = data_shape[final_index]
                        m_stride[final_index] = 1
                        fshape_indices.append(final_index)
                        final_index += 1
                elif mask &new_axis_mask:
                    fshape_indices.append(-1)
                elif not mask & new_axis_mask:
                    if final_index == len(m_begin):
                        break
                    if mask & begin_mask:
                        m_begin[final_index] = data_shape[final_index] \
                                                     if stride[index] < 0 else 0
                    elif begin[index]:
                        m_begin[final_index] = begin[index]
                    if mask & end_mask:
                        m_end[final_index] = 0 if stride[index] < 0 \
                                                 else data_shape[final_index]
                    elif end[index]:
                        m_end[final_index] = end[index]
                    m_stride[final_index] = stride[index]
                    if mask & shrink_axis_mask:
                        #Tensorflow make axis with shrink_axis_mask as dimension 1
                        m_begin[final_index] = data_shape[final_index] + begin[index] \
                                                 if begin[index] < 0 else begin[index]
                        m_end[final_index] = begin[index] + 1
                        m_stride[final_index] = 1
                        fshape_indices.append(-2)
                    else:
                        fshape_indices.append(final_index)

                    final_index += 1
            return m_begin, m_end, m_stride, fshape_indices

        fshape_indices = None
        if begin_mask or end_mask or ellipsis_mask or new_axis_mask or shrink_axis_mask:
            #begin from -1 to 3 (-1 -> 3) ;end 0;stride=1
            begin, end, stride, fshape_indices = _transform_mask(stride_dim, ellipsis_mask)
        out = _op.strided_slice(inputs[0], begin=begin, end=end, strides=stride)
        #error
        out_shape = _infer_shape(out, mod=mod)   <-----------error
        if not fshape_indices:
            fshape_indices = range(len(out_shape))

        #Create final output shape.
        final_output = []
        for gather_index in fshape_indices:
            if gather_index == -1:
                final_output.append(1)
            elif gather_index == -2:
                pass
            else:
                final_output.append(out_shape[gather_index])

        if not final_output:
            return out
        return _op.reshape(out, newshape=tuple(final_output))
    return _impl

error info:

an internal invariant was violated while typechecking your program
Check failed: begin_v < end_v (3 vs. 0) : strided_slice get empty slice at axis 0
......
File xxxxxx  in _impl
        out_shape = _infer_shape(out, mod=mod)
File xxxxx  in infer_shape
        out_type = infer_type(inputs, mod=mod)
File xxxxx  in infer_type
       new_mod = IRModule.from_expr(node)

Traceback:
[bt] (8) tvm::IRModuleNode::FromExpr(xxxx)
[bt] (7) tvm::IRModuleNode::Add(xxx)
[bt] (6) tvm::RunTypeCheck(xx)
[bt] (5) tvm::relay::inferType(xxx)
[bt] (4) tvm::relay::TypeInference::infer(xxx)
[bt] (3) tvm::relay::TypeSolver::solve(xxx)

Well, I can only guess from strided_slice get empty slice at axis 0 that some attribute maybe empty. Perhaps you can try to run _infer_value for every input or attribute to see it.

I run _infer_value for inputs[0] and print it.

inputs[0]  infer_value: [1 8 1 64]

I run _infer_value for other inputs(inputs[1] .et) and error

#should error
AssertionError: All inputs to infer must be available in params.

print attr of input:

attr:
{'begin_mask':0,'end_mask':0,'ellipsis_mask':0,'shrink_axis_mask':1,'new_axis_mask':0
'_output_shapes':[[]],'T':tf.int32,'_node_name':'strided_slice_4','_target_layout':None,
'_input_shape':{var(Shape_3,ty=TensorType([4],int32)): [4],
var(strided_slice_4/stack,ty=TensorType([1],int32):[1],
var(strided_slice_4/stack_1,ty=TensorType([1],int32):[1],
var(strided_slice_4/stack_2,ty=TensorType([1],int32):[1],
},}
begin = [-1] end = [0] stride=[1]
after _transform_mask:
begin = [3] end = [0] stride=[1]

inputs[0] is not empty,there seems to be no problem.error because begin(=3) > end(0)?

I also met this problem using the latest code. Have you solved this problem? thanks

This error:

strided_slice(%307, meta[relay.Constant][0], meta[relay.Constant][1], meta[relay.Constant][2], begin=[0, 3, 0, 0], end=[1, 0, 1, 512], strides=[1, 1, 1, 1]) 
an internal invariant was violated while typechecking your program [12:17:52]  
/root/worker/tvm/src/relay/op/tensor/transform.cc:2074: 
Check failed: begin_v <= end_v (3 vs. 0) : strided_slice get invalid slice at axis 1